Skip to content

Commit 48353bf

Browse files
committed
Merge remote-tracking branch 'origin/master' into out_of_range
2 parents d65169c + f057808 commit 48353bf

Some content is hidden

Large Commits have some content hidden by default. Use the searchbox below for content that may be hidden.

46 files changed

+1417
-546
lines changed

cmake/build-info.cmake

Lines changed: 8 additions & 2 deletions
Original file line numberDiff line numberDiff line change
@@ -41,14 +41,20 @@ endif()
4141

4242
if(MSVC)
4343
set(BUILD_COMPILER "${CMAKE_C_COMPILER_ID} ${CMAKE_C_COMPILER_VERSION}")
44-
set(BUILD_TARGET ${CMAKE_VS_PLATFORM_NAME})
44+
if (CMAKE_VS_PLATFORM_NAME)
45+
set(BUILD_TARGET ${CMAKE_VS_PLATFORM_NAME})
46+
else()
47+
set(BUILD_TARGET "${CMAKE_SYSTEM_NAME} ${CMAKE_SYSTEM_PROCESSOR}")
48+
endif()
4549
else()
4650
execute_process(
47-
COMMAND sh -c "\"$@\" --version | head -1" _ ${CMAKE_C_COMPILER}
51+
COMMAND ${CMAKE_C_COMPILER} --version
4852
OUTPUT_VARIABLE OUT
4953
OUTPUT_STRIP_TRAILING_WHITESPACE
5054
)
55+
string(REGEX REPLACE " *\n.*" "" OUT "${OUT}")
5156
set(BUILD_COMPILER ${OUT})
57+
5258
execute_process(
5359
COMMAND ${CMAKE_C_COMPILER} -dumpmachine
5460
OUTPUT_VARIABLE OUT

common/CMakeLists.txt

Lines changed: 3 additions & 1 deletion
Original file line numberDiff line numberDiff line change
@@ -39,7 +39,9 @@ add_custom_command(
3939
COMMENT "Generating build details from Git"
4040
COMMAND ${CMAKE_COMMAND} -DMSVC=${MSVC} -DCMAKE_C_COMPILER_VERSION=${CMAKE_C_COMPILER_VERSION}
4141
-DCMAKE_C_COMPILER_ID=${CMAKE_C_COMPILER_ID} -DCMAKE_VS_PLATFORM_NAME=${CMAKE_VS_PLATFORM_NAME}
42-
-DCMAKE_C_COMPILER=${CMAKE_C_COMPILER} -P "${CMAKE_CURRENT_SOURCE_DIR}/cmake/build-info-gen-cpp.cmake"
42+
-DCMAKE_C_COMPILER=${CMAKE_C_COMPILER}
43+
-DCMAKE_SYSTEM_NAME=${CMAKE_SYSTEM_NAME} -DCMAKE_SYSTEM_PROCESSOR=${CMAKE_SYSTEM_PROCESSOR}
44+
-P "${CMAKE_CURRENT_SOURCE_DIR}/cmake/build-info-gen-cpp.cmake"
4345
WORKING_DIRECTORY "${CMAKE_CURRENT_SOURCE_DIR}/.."
4446
DEPENDS "${CMAKE_CURRENT_SOURCE_DIR}/build-info.cpp.in" ${GIT_INDEX}
4547
VERBATIM

common/arg.cpp

Lines changed: 6 additions & 15 deletions
Original file line numberDiff line numberDiff line change
@@ -217,13 +217,11 @@ struct curl_slist_ptr {
217217
#define CURL_MAX_RETRY 3
218218
#define CURL_RETRY_DELAY_SECONDS 2
219219

220-
static bool curl_perform_with_retry(const std::string & url, CURL * curl, int max_attempts, int retry_delay_seconds) {
220+
static bool curl_perform_with_retry(const std::string & url, CURL * curl, int max_attempts, int retry_delay_seconds, const char * method_name) {
221221
int remaining_attempts = max_attempts;
222-
char * method = nullptr;
223-
curl_easy_getinfo(curl, CURLINFO_EFFECTIVE_METHOD, &method);
224222

225223
while (remaining_attempts > 0) {
226-
LOG_INF("%s: %s %s (attempt %d of %d)...\n", __func__ , method, url.c_str(), max_attempts - remaining_attempts + 1, max_attempts);
224+
LOG_INF("%s: %s %s (attempt %d of %d)...\n", __func__ , method_name, url.c_str(), max_attempts - remaining_attempts + 1, max_attempts);
227225

228226
CURLcode res = curl_easy_perform(curl);
229227
if (res == CURLE_OK) {
@@ -287,24 +285,17 @@ static bool common_download_file_single(const std::string & url, const std::stri
287285
try {
288286
metadata_in >> metadata;
289287
LOG_DBG("%s: previous metadata file found %s: %s\n", __func__, metadata_path.c_str(), metadata.dump().c_str());
290-
if (metadata.contains("url") && metadata.at("url").is_string()) {
291-
auto previous_url = metadata.at("url").get<std::string>();
292-
if (previous_url != url) {
293-
LOG_ERR("%s: Model URL mismatch: %s != %s\n", __func__, url.c_str(), previous_url.c_str());
294-
return false;
295-
}
296-
}
297288
if (metadata.contains("etag") && metadata.at("etag").is_string()) {
298289
etag = metadata.at("etag");
299290
}
300291
if (metadata.contains("lastModified") && metadata.at("lastModified").is_string()) {
301292
last_modified = metadata.at("lastModified");
302293
}
303294
} catch (const nlohmann::json::exception & e) {
304-
LOG_ERR("%s: error reading metadata file %s: %s\n", __func__, metadata_path.c_str(), e.what());
305-
return false;
295+
LOG_ERR("%s: error reading metadata file %s: %s\n", __func__, metadata_path.c_str(), e.what());
306296
}
307297
}
298+
// if we cannot open the metadata file, we assume that the downloaded file is not valid (etag and last-modified are left empty, so we will download it again)
308299
} else {
309300
LOG_INF("%s: no previous model file found %s\n", __func__, path.c_str());
310301
}
@@ -350,7 +341,7 @@ static bool common_download_file_single(const std::string & url, const std::stri
350341

351342
// we only allow retrying once for HEAD requests
352343
// this is for the use case of using running offline (no internet), retrying can be annoying
353-
bool was_perform_successful = curl_perform_with_retry(url, curl.get(), 1, 0);
344+
bool was_perform_successful = curl_perform_with_retry(url, curl.get(), 1, 0, "HEAD");
354345
if (!was_perform_successful) {
355346
head_request_ok = false;
356347
}
@@ -432,7 +423,7 @@ static bool common_download_file_single(const std::string & url, const std::stri
432423
// start the download
433424
LOG_INF("%s: trying to download model from %s to %s (server_etag:%s, server_last_modified:%s)...\n", __func__,
434425
llama_download_hide_password_in_url(url).c_str(), path.c_str(), headers.etag.c_str(), headers.last_modified.c_str());
435-
bool was_perform_successful = curl_perform_with_retry(url, curl.get(), CURL_MAX_RETRY, CURL_RETRY_DELAY_SECONDS);
426+
bool was_perform_successful = curl_perform_with_retry(url, curl.get(), CURL_MAX_RETRY, CURL_RETRY_DELAY_SECONDS, "GET");
436427
if (!was_perform_successful) {
437428
return false;
438429
}

convert_hf_to_gguf.py

Lines changed: 27 additions & 3 deletions
Original file line numberDiff line numberDiff line change
@@ -1899,7 +1899,10 @@ def prepare_tensors(self):
18991899
raise ValueError(f"Unprocessed experts: {experts}")
19001900

19011901

1902-
@ModelBase.register("LlavaForConditionalGeneration")
1902+
@ModelBase.register(
1903+
"LlavaForConditionalGeneration", # pixtral
1904+
"Mistral3ForConditionalGeneration", # mistral small 3.1
1905+
)
19031906
class LlavaVisionModel(VisionModel):
19041907
img_break_tok_id = -1
19051908

@@ -1908,17 +1911,38 @@ def __init__(self, *args, **kwargs):
19081911
if self.hparams["model_type"] == "pixtral":
19091912
# layer_norm_eps is not in config.json, it is hard-coded in modeling_pixtral.py
19101913
self.hparams["layer_norm_eps"] = self.hparams.get("layer_norm_eps", 1e-5)
1911-
self.img_break_tok_id = 12 # see tokenizer_config.json
1914+
self.img_break_tok_id = self.get_token_id("[IMG_BREAK]")
1915+
logger.info(f"Image break token id: {self.img_break_tok_id}")
19121916
else:
19131917
raise ValueError(f"Unsupported model type: {self.hparams['model_type']}")
19141918

1919+
def get_token_id(self, token: str) -> int:
1920+
tokenizer_config_file = self.dir_model / 'tokenizer_config.json'
1921+
with open(tokenizer_config_file, "r", encoding="utf-8") as f:
1922+
added_tokens_decoder = json.load(f)['added_tokens_decoder']
1923+
for id_, token_data in added_tokens_decoder.items():
1924+
if token_data["content"] == token:
1925+
return int(id_)
1926+
raise ValueError(f"Token '{token}' not found in tokenizer config.")
1927+
19151928
def set_gguf_parameters(self):
19161929
super().set_gguf_parameters()
19171930
hparams = self.hparams
19181931
if hparams["model_type"] == "pixtral":
19191932
self.gguf_writer.add_vision_projector_type(gguf.VisionProjectorType.PIXTRAL)
19201933
self.gguf_writer.add_vision_attention_layernorm_eps(hparams["layer_norm_eps"])
1921-
self.gguf_writer.add_vision_use_silu(True)
1934+
1935+
# hidden_act
1936+
if hparams["hidden_act"] == "silu":
1937+
self.gguf_writer.add_vision_use_silu(True)
1938+
elif hparams["hidden_act"] == "gelu":
1939+
self.gguf_writer.add_vision_use_gelu(True)
1940+
else:
1941+
raise ValueError(f"Unsupported hidden_act: {hparams['hidden_act']}")
1942+
1943+
# spatial_merge_size
1944+
if "spatial_merge_size" in self.global_config:
1945+
self.gguf_writer.add_vision_spatial_merge_size(self.global_config["spatial_merge_size"])
19221946

19231947
def modify_tensors(self, data_torch: Tensor, name: str, bid: int | None) -> Iterable[tuple[str, Tensor]]:
19241948
del bid # unused

examples/llava/README.md

Lines changed: 4 additions & 0 deletions
Original file line numberDiff line numberDiff line change
@@ -34,6 +34,9 @@ llama-mtmd-cli -hf ggml-org/SmolVLM2-500M-Video-Instruct-GGUF
3434

3535
# Pixtral 12B
3636
llama-mtmd-cli -hf ggml-org/pixtral-12b-GGUF
37+
38+
# Mistral Small 3.1 24B (IQ2_M quantization)
39+
llama-mtmd-cli -hf ggml-org/Mistral-Small-3.1-24B-Instruct-2503-GGUF --chat-template mistral-v7
3740
```
3841

3942
## How it works and what is `mmproj`?
@@ -73,3 +76,4 @@ For the following models, you can use `convert_hf_to_gguf.py`with `--mmproj` fla
7376
- SmolVLM (from [HuggingFaceTB](https://huggingface.co/HuggingFaceTB))
7477
- SmolVLM2 (from [HuggingFaceTB](https://huggingface.co/HuggingFaceTB))
7578
- [Pixtral 12B](https://huggingface.co/mistral-community/pixtral-12b) - only works with `transformers`-compatible checkpoint
79+
- [Mistral Small 3.1 24B](https://huggingface.co/mistralai/Mistral-Small-3.1-24B-Instruct-2503)

examples/llava/clip-impl.h

Lines changed: 3 additions & 0 deletions
Original file line numberDiff line numberDiff line change
@@ -31,6 +31,7 @@
3131
#define KEY_FEATURE_LAYER "clip.vision.feature_layer"
3232
#define KEY_PROJ_SCALE_FACTOR "clip.vision.projector.scale_factor"
3333
#define KEY_PROJ_TYPE "clip.projector_type"
34+
#define KEY_SPATIAL_MERGE_SIZE "clip.vision.spatial_merge_size"
3435

3536
#define KEY_USE_GLU_MLP "clip.use_glu_mlp" // for qwen2.5vl
3637
#define KEY_USE_RMS_NORM "clip.use_rms_norm" // for qwen2.5vl
@@ -68,9 +69,11 @@
6869
#define TN_MVLM_PROJ_BLOCK "mm.model.mb_block.%d.block.%d.%s"
6970
#define TN_MVLM_PROJ_PEG "mm.model.peg.%d.%s"
7071
#define TN_IMAGE_NEWLINE "model.image_newline"
72+
#define TN_MM_INP_NORM "mm.input_norm.weight"
7173
#define TN_MM_INP_PROJ "mm.input_projection.weight" // gemma3
7274
#define TN_MM_SOFT_EMB_N "mm.soft_emb_norm.weight" // gemma3
7375
#define TN_MM_PROJECTOR "mm.model.fc.weight" // idefics3
76+
#define TN_MM_PATCH_MERGER "mm.patch_merger.weight" // mistral small 3.1
7477
#define TN_TOK_IMG_BREAK "v.token_embd.img_break" // pixtral
7578

7679
// mimicpmv

examples/llava/clip.cpp

Lines changed: 59 additions & 13 deletions
Original file line numberDiff line numberDiff line change
@@ -172,6 +172,7 @@ struct clip_hparams {
172172
std::unordered_set<int32_t> vision_feature_layer;
173173
int32_t attn_window_size = 0;
174174
int32_t n_wa_pattern = 0;
175+
int32_t spatial_merge_size = 0;
175176
};
176177

177178
struct clip_layer {
@@ -232,6 +233,7 @@ struct clip_vision_model {
232233
struct ggml_tensor * projection;
233234

234235
// LLaVA projection
236+
struct ggml_tensor * mm_input_norm_w = nullptr;
235237
struct ggml_tensor * mm_0_w = nullptr;
236238
struct ggml_tensor * mm_0_b = nullptr;
237239
struct ggml_tensor * mm_2_w = nullptr;
@@ -311,6 +313,7 @@ struct clip_vision_model {
311313

312314
// pixtral
313315
struct ggml_tensor * token_embd_img_break = nullptr;
316+
struct ggml_tensor * mm_patch_merger_w = nullptr;
314317
};
315318

316319
struct clip_ctx {
@@ -637,6 +640,7 @@ static ggml_cgraph * clip_image_build_graph_pixtral(clip_ctx * ctx, const clip_i
637640
const int d_head = hidden_size / n_head;
638641
const int n_layer = hparams.n_layer;
639642
const float eps = hparams.eps;
643+
const int n_merge = hparams.spatial_merge_size;
640644

641645
struct ggml_init_params params = {
642646
/*.mem_size =*/ ctx->buf_compute_meta.size(),
@@ -721,7 +725,13 @@ static ggml_cgraph * clip_image_build_graph_pixtral(clip_ctx * ctx, const clip_i
721725
{
722726
ggml_tensor * gate_proj = ggml_mul_mat(ctx0, model.layers[il].ff_gate_w, cur);
723727
ggml_tensor * up_proj = ggml_mul_mat(ctx0, model.layers[il].ff_up_w, cur);
724-
gate_proj = ggml_silu(ctx0, gate_proj); // pixtral uses silu
728+
if (ctx->use_silu) {
729+
gate_proj = ggml_silu(ctx0, gate_proj);
730+
} else if (ctx->use_gelu) {
731+
gate_proj = ggml_gelu(ctx0, gate_proj);
732+
} else {
733+
GGML_ABORT("Pixtral: Unsupported activation");
734+
}
725735
cur = ggml_mul(ctx0, up_proj, gate_proj);
726736
cur = ggml_mul_mat(ctx0, model.layers[il].ff_down_w, cur);
727737
}
@@ -732,14 +742,42 @@ static ggml_cgraph * clip_image_build_graph_pixtral(clip_ctx * ctx, const clip_i
732742
embeddings = cur;
733743
}
734744

735-
// LlavaMultiModalProjector (with GELU activation)
745+
// mistral small 3.1 patch merger
746+
// ref: https://github.com/huggingface/transformers/blob/7a3e208892c06a5e278144eaf38c8599a42f53e7/src/transformers/models/mistral3/modeling_mistral3.py#L67
747+
if (model.mm_patch_merger_w) {
748+
GGML_ASSERT(hparams.spatial_merge_size > 0);
749+
750+
ggml_tensor * cur = embeddings;
751+
cur = ggml_mul(ctx0, ggml_rms_norm(ctx0, cur, eps), model.mm_input_norm_w);
752+
753+
// reshape image tokens to 2D grid
754+
cur = ggml_reshape_3d(ctx0, cur, hidden_size, n_patches_x, n_patches_y);
755+
cur = ggml_permute(ctx0, cur, 2, 0, 1, 3); // [x, y, hidden_size]
756+
cur = ggml_cont(ctx0, cur);
757+
758+
// torch.nn.functional.unfold is just an im2col under the hood
759+
// we just need a dummy kernel to make it work
760+
ggml_tensor * kernel = ggml_view_3d(ctx0, cur, n_merge, n_merge, cur->ne[2], 0, 0, 0);
761+
cur = ggml_im2col(ctx0, kernel, cur, n_merge, n_merge, 0, 0, 1, 1, true, inp->type);
762+
763+
// project to hidden_size
764+
cur = ggml_reshape_2d(ctx0, cur, cur->ne[0], cur->ne[1] * cur->ne[2]);
765+
cur = ggml_mul_mat(ctx0, model.mm_patch_merger_w, cur);
766+
embeddings = cur;
767+
}
768+
769+
// LlavaMultiModalProjector (always using GELU activation)
736770
{
737771
embeddings = ggml_mul_mat(ctx0, model.mm_1_w, embeddings);
738-
embeddings = ggml_add(ctx0, embeddings, model.mm_1_b);
772+
if (model.mm_1_b) {
773+
embeddings = ggml_add(ctx0, embeddings, model.mm_1_b);
774+
}
739775

740776
embeddings = ggml_gelu(ctx0, embeddings);
741777
embeddings = ggml_mul_mat(ctx0, model.mm_2_w, embeddings);
742-
embeddings = ggml_add(ctx0, embeddings, model.mm_2_b);
778+
if (model.mm_2_b) {
779+
embeddings = ggml_add(ctx0, embeddings, model.mm_2_b);
780+
}
743781
}
744782

745783
// arrangement of the [IMG_BREAK] token
@@ -749,11 +787,14 @@ static ggml_cgraph * clip_image_build_graph_pixtral(clip_ctx * ctx, const clip_i
749787
// and then concatenate the [IMG_BREAK] token to the end of each row, aka n_patches_per_row dimension
750788
// after the concatenation, we have a tensor with shape [hidden_size, n_patches_per_row + 1, n_rows]
751789

790+
const int p_y = n_merge > 0 ? n_patches_y / n_merge : n_patches_y;
791+
const int p_x = n_merge > 0 ? n_patches_x / n_merge : n_patches_x;
792+
const int p_total = p_x * p_y;
752793
const int n_embd_text = embeddings->ne[0];
753-
const int n_tokens_output = num_patches + n_patches_y - 1; // one [IMG_BREAK] per row, except the last row
794+
const int n_tokens_output = p_total + p_y - 1; // one [IMG_BREAK] per row, except the last row
754795

755-
ggml_tensor * cur = ggml_reshape_3d(ctx0, embeddings, n_embd_text, n_patches_x, n_patches_y);
756-
ggml_tensor * tok = ggml_new_tensor_3d(ctx0, embeddings->type, n_embd_text, 1, n_patches_y);
796+
ggml_tensor * cur = ggml_reshape_3d(ctx0, embeddings, n_embd_text, p_x, p_y);
797+
ggml_tensor * tok = ggml_new_tensor_3d(ctx0, embeddings->type, n_embd_text, 1, p_y);
757798
tok = ggml_scale(ctx0, tok, 0.0); // clear the tensor
758799
tok = ggml_add(ctx0, tok, model.token_embd_img_break);
759800
cur = ggml_concat(ctx0, cur, tok, 1);
@@ -1734,6 +1775,7 @@ struct clip_model_loader {
17341775
case PROJECTOR_TYPE_PIXTRAL:
17351776
{
17361777
hparams.rope_theta = 10000.0f;
1778+
get_u32(KEY_SPATIAL_MERGE_SIZE, hparams.spatial_merge_size, false);
17371779
} break;
17381780
case PROJECTOR_TYPE_QWEN25VL:
17391781
{
@@ -1957,11 +1999,14 @@ struct clip_model_loader {
19571999
case PROJECTOR_TYPE_PIXTRAL:
19582000
{
19592001
vision_model.mm_1_w = get_tensor(string_format(TN_LLAVA_PROJ, 1, "weight"));
1960-
vision_model.mm_1_b = get_tensor(string_format(TN_LLAVA_PROJ, 1, "bias"));
2002+
vision_model.mm_1_b = get_tensor(string_format(TN_LLAVA_PROJ, 1, "bias"), false);
19612003
vision_model.mm_2_w = get_tensor(string_format(TN_LLAVA_PROJ, 2, "weight"));
1962-
vision_model.mm_2_b = get_tensor(string_format(TN_LLAVA_PROJ, 2, "bias"));
2004+
vision_model.mm_2_b = get_tensor(string_format(TN_LLAVA_PROJ, 2, "bias"), false);
19632005
// [IMG_BREAK] token embedding
19642006
vision_model.token_embd_img_break = get_tensor(TN_TOK_IMG_BREAK);
2007+
// for mistral small 3.1
2008+
vision_model.mm_input_norm_w = get_tensor(TN_MM_INP_NORM, false);
2009+
vision_model.mm_patch_merger_w = get_tensor(TN_MM_PATCH_MERGER, false);
19652010
} break;
19662011
default:
19672012
GGML_ASSERT(false && "unknown projector type");
@@ -2516,7 +2561,7 @@ struct llava_uhd {
25162561

25172562
// no pinpoints, dynamically calculate the grid size (e.g. minicpmv)
25182563

2519-
auto best_size = get_best_resize(original_size, slice_size, patch_size, has_slices);
2564+
auto best_size = get_best_resize(original_size, slice_size, patch_size, !has_slices);
25202565
res.overview_size = best_size;
25212566

25222567
if (!has_slices) {
@@ -2926,8 +2971,9 @@ int clip_n_output_tokens(const struct clip_ctx * ctx, struct clip_image_f32 * im
29262971
} else if (ctx->proj_type == PROJECTOR_TYPE_IDEFICS3) {
29272972
n_patches /= ctx->vision_model.hparams.proj_scale_factor;
29282973
} else if (ctx->proj_type == PROJECTOR_TYPE_PIXTRAL) {
2929-
int n_patches_x = img->nx / params.patch_size;
2930-
int n_patches_y = img->ny / params.patch_size;
2974+
int n_merge = ctx->vision_model.hparams.spatial_merge_size;
2975+
int n_patches_x = img->nx / params.patch_size / (n_merge > 0 ? n_merge : 1);
2976+
int n_patches_y = img->ny / params.patch_size / (n_merge > 0 ? n_merge : 1);
29312977
n_patches = n_patches_y*n_patches_x + n_patches_y - 1; // + one [IMG_BREAK] per row, except the last row
29322978
}
29332979

@@ -3484,7 +3530,7 @@ int clip_n_mmproj_embd(const struct clip_ctx * ctx) {
34843530
return ctx->vision_model.mm_model_peg_0_b->ne[0];
34853531
case PROJECTOR_TYPE_MLP:
34863532
case PROJECTOR_TYPE_PIXTRAL:
3487-
return ctx->vision_model.mm_2_b->ne[0];
3533+
return ctx->vision_model.mm_2_w->ne[1];
34883534
case PROJECTOR_TYPE_MLP_NORM:
34893535
return ctx->vision_model.mm_3_b->ne[0];
34903536
case PROJECTOR_TYPE_MINICPMV:

examples/llava/mtmd-cli.cpp

Lines changed: 1 addition & 0 deletions
Original file line numberDiff line numberDiff line change
@@ -96,6 +96,7 @@ struct mtmd_cli_context {
9696
LOG_ERR("Model does not have chat template.\n");
9797
LOG_ERR(" For old llava models, you may need to use '--chat-template vicuna'\n");
9898
LOG_ERR(" For MobileVLM models, use '--chat-template deepseek'\n");
99+
LOG_ERR(" For Mistral Small 3.1, use '--chat-template mistral-v7'\n");
99100
exit(1);
100101
}
101102

examples/llava/tests.sh

Lines changed: 1 addition & 0 deletions
Original file line numberDiff line numberDiff line change
@@ -59,6 +59,7 @@ add_test "llama-mtmd-cli" "ggml-org/Qwen2.5-VL-3B-Instruct-GGUF:Q4_K_M"
5959

6060
# to test the big models, run: ./tests.sh big
6161
add_test_big "llama-mtmd-cli" "ggml-org/pixtral-12b-GGUF:Q4_K_M"
62+
add_test_big "llama-mtmd-cli" "ggml-org/Mistral-Small-3.1-24B-Instruct-2503-GGUF" "mistral-v7"
6263

6364
# these models always give the wrong answer, not sure why
6465
# add_test "llama-mtmd-cli" "ggml-org/SmolVLM-Instruct-GGUF:Q4_K_M"

ggml/CMakeLists.txt

Lines changed: 24 additions & 0 deletions
Original file line numberDiff line numberDiff line change
@@ -360,3 +360,27 @@ write_basic_package_version_file(
360360
install(FILES ${CMAKE_CURRENT_BINARY_DIR}/ggml-config.cmake
361361
${CMAKE_CURRENT_BINARY_DIR}/ggml-version.cmake
362362
DESTINATION ${CMAKE_INSTALL_LIBDIR}/cmake/ggml)
363+
364+
if (MSVC)
365+
set(MSVC_WARNING_FLAGS
366+
/wd4005 # Macro redefinition
367+
/wd4244 # Conversion from one type to another type, possible loss of data
368+
/wd4267 # Conversion from 'size_t' to a smaller type, possible loss of data
369+
)
370+
function(disable_msvc_warnings target_name)
371+
if(TARGET ${target_name})
372+
target_compile_options(${target_name} PRIVATE ${MSVC_WARNING_FLAGS})
373+
endif()
374+
endfunction()
375+
376+
disable_msvc_warnings(ggml-base)
377+
disable_msvc_warnings(ggml)
378+
disable_msvc_warnings(ggml-cpu)
379+
disable_msvc_warnings(ggml-cpu-x64)
380+
disable_msvc_warnings(ggml-cpu-sse42)
381+
disable_msvc_warnings(ggml-cpu-sandybridge)
382+
disable_msvc_warnings(ggml-cpu-haswell)
383+
disable_msvc_warnings(ggml-cpu-skylakex)
384+
disable_msvc_warnings(ggml-cpu-icelake)
385+
disable_msvc_warnings(ggml-cpu-alderlake)
386+
endif()

0 commit comments

Comments
 (0)