Skip to content

Commit 7bed317

Browse files
authored
models : fix the attn_factor for mistral3 graphs + improve consistency (ggml-org#17945)
* models : fix the attn_factor for mistral3 graphs * cont : rework attn_factor correction logic * cont : make deepseek2 consistent * cont : add TODO * cont : special-case DSv2 * cont : revert Mistral 3 Large changes * cont : fix DS2 to use the original attn_factor * cont : minor comments
1 parent dcb7d17 commit 7bed317

File tree

7 files changed

+78
-33
lines changed

7 files changed

+78
-33
lines changed

convert_hf_to_gguf.py

Lines changed: 8 additions & 0 deletions
Original file line numberDiff line numberDiff line change
@@ -7286,6 +7286,10 @@ def set_gguf_parameters(self):
72867286
self.gguf_writer.add_rope_scaling_type(gguf.RopeScalingType.YARN)
72877287
self.gguf_writer.add_rope_scaling_factor(rope_scaling["factor"])
72887288
self.gguf_writer.add_rope_scaling_orig_ctx_len(rope_scaling["original_max_position_embeddings"])
7289+
7290+
# [TAG_DEEPSEEK2_YARN_LOG_MUL_FIX]
7291+
# note: for legacy reasons, this is not consistent with the other usages of self.gguf_writer.add_rope_scaling_yarn_log_mul
7292+
# ref https://github.com/ggml-org/llama.cpp/pull/17945
72897293
self.gguf_writer.add_rope_scaling_yarn_log_mul(0.1 * rope_scaling["mscale_all_dim"])
72907294

72917295
_experts: list[dict[str, Tensor]] | None = None
@@ -10041,6 +10045,10 @@ def set_gguf_parameters(self):
1004110045
MistralModel.set_mistral_config(self.gguf_writer, self.hparams)
1004210046
yarn_params = self.hparams["yarn"]
1004310047
self.gguf_writer.add_attn_temperature_length(yarn_params["original_max_position_embeddings"])
10048+
10049+
# [TAG_DEEPSEEK2_YARN_LOG_MUL_FIX]
10050+
# note: for legacy reasons, this is not consistent with the other usages of self.gguf_writer.add_rope_scaling_yarn_log_mul
10051+
# ref https://github.com/ggml-org/llama.cpp/pull/17945
1004410052
self.gguf_writer.add_rope_scaling_yarn_log_mul(0.1) # mscale_all_dim * 0.1
1004510053

1004610054
def modify_tensors(self, data_torch: Tensor, name: str, bid: int | None):

src/llama-graph.cpp

Lines changed: 1 addition & 1 deletion
Original file line numberDiff line numberDiff line change
@@ -574,7 +574,7 @@ llm_graph_context::llm_graph_context(const llm_graph_params & params) :
574574
freq_base (cparams.rope_freq_base),
575575
freq_scale (cparams.rope_freq_scale),
576576
ext_factor (cparams.yarn_ext_factor),
577-
attn_factor (cparams.yarn_attn_factor),
577+
attn_factor (llama_hparams::yarn_attn_factor_adjust(cparams.yarn_attn_factor, cparams.rope_freq_scale, cparams.yarn_ext_factor)),
578578
beta_fast (cparams.yarn_beta_fast),
579579
beta_slow (cparams.yarn_beta_slow),
580580
norm_eps (hparams.f_norm_eps),

src/llama-hparams.cpp

Lines changed: 12 additions & 0 deletions
Original file line numberDiff line numberDiff line change
@@ -1,7 +1,9 @@
11
#include "llama-hparams.h"
22

33
#include "ggml.h"
4+
45
#include <cassert>
6+
#include <cmath>
57

68
void llama_hparams::set_swa_pattern(uint32_t n_pattern, bool dense_first) {
79
if (dense_first) {
@@ -229,3 +231,13 @@ bool llama_hparams::is_masked_swa(uint32_t n_swa, llama_swa_type swa_type, llama
229231

230232
return false;
231233
}
234+
235+
float llama_hparams::yarn_attn_factor_adjust(float attn_factor, float freq_scale, float ext_factor) {
236+
GGML_ASSERT(ext_factor >= 0.0f);
237+
238+
if (ext_factor != 0.0f) {
239+
attn_factor *= 1.0f / (1.0f + 0.1f * logf(1.0f / freq_scale));
240+
}
241+
242+
return attn_factor;
243+
}

src/llama-hparams.h

Lines changed: 8 additions & 1 deletion
Original file line numberDiff line numberDiff line change
@@ -107,6 +107,7 @@ struct llama_hparams {
107107
float rope_freq_base_train_swa;
108108
float rope_freq_scale_train;
109109
float rope_freq_scale_train_swa;
110+
110111
uint32_t n_ctx_orig_yarn;
111112
float rope_yarn_log_mul = 0.0f;
112113

@@ -267,7 +268,13 @@ struct llama_hparams {
267268
// TODO: think of a better place for this function
268269
// TODO: pack the SWA params in a struct?
269270
static bool is_masked_swa(uint32_t n_swa, llama_swa_type swa_type, llama_pos p0, llama_pos p1);
271+
272+
// when YARN is applied with yarn_ext_factor != 0.0f, we need to cancel this factor:
273+
// https://github.com/ggml-org/llama.cpp/blob/a81a569577cc38b32558958b048228150be63eae/ggml/src/ggml-cpu/ops.cpp#L5541-L5544
274+
//
275+
// ref: https://github.com/ggml-org/llama.cpp/discussions/7416
276+
// https://github.com/ggml-org/llama.cpp/pull/17945
277+
static float yarn_attn_factor_adjust(float attn_factor, float freq_scale, float ext_factor);
270278
};
271279

272280
static_assert(std::is_trivially_copyable<llama_hparams>::value, "llama_hparams must be trivially copyable");
273-

src/llama-kv-cache.cpp

Lines changed: 4 additions & 9 deletions
Original file line numberDiff line numberDiff line change
@@ -1369,9 +1369,10 @@ ggml_tensor * llama_kv_cache::build_rope_shift(
13691369
float freq_scale) const {
13701370
const auto & n_ctx_orig = cparams.n_ctx_orig_yarn;
13711371

1372-
const auto & yarn_ext_factor = cparams.yarn_ext_factor;
1373-
const auto & yarn_beta_fast = cparams.yarn_beta_fast;
1374-
const auto & yarn_beta_slow = cparams.yarn_beta_slow;
1372+
const auto & yarn_ext_factor = cparams.yarn_ext_factor;
1373+
const auto & yarn_beta_fast = cparams.yarn_beta_fast;
1374+
const auto & yarn_beta_slow = cparams.yarn_beta_slow;
1375+
const auto & yarn_attn_factor = llama_hparams::yarn_attn_factor_adjust(cparams.yarn_attn_factor, cparams.rope_freq_scale, cparams.yarn_ext_factor);
13751376

13761377
const auto & n_rot = hparams.n_rot;
13771378
const auto & rope_type = hparams.rope_type == LLAMA_ROPE_TYPE_MROPE || hparams.rope_type == LLAMA_ROPE_TYPE_IMROPE
@@ -1382,12 +1383,6 @@ ggml_tensor * llama_kv_cache::build_rope_shift(
13821383
? LLAMA_ROPE_TYPE_NEOX
13831384
: hparams.rope_type;
13841385

1385-
// See llm_build_deepseek2() for why attn_factor has to be scaled for YaRN RoPE to work correctly.
1386-
// See https://github.com/ggerganov/llama.cpp/discussions/7416 for detailed explanation.
1387-
const float yarn_attn_factor = model.arch == LLM_ARCH_DEEPSEEK2
1388-
? 1.0f / (1.0f + 0.1f * logf(1.0f / freq_scale))
1389-
: cparams.yarn_attn_factor;
1390-
13911386
ggml_tensor * tmp;
13921387

13931388
if (ggml_is_quantized(cur->type)) {

src/llama-model.cpp

Lines changed: 36 additions & 17 deletions
Original file line numberDiff line numberDiff line change
@@ -1635,7 +1635,12 @@ void llama_model::load_hparams(llama_model_loader & ml) {
16351635
// that have no expert_gating_func model parameter set
16361636
hparams.expert_gating_func = LLAMA_EXPERT_GATING_FUNC_TYPE_SOFTMAX;
16371637
}
1638-
ml.get_key(LLM_KV_ROPE_SCALING_YARN_LOG_MUL, hparams.rope_yarn_log_mul, false);
1638+
1639+
if (ml.get_key(LLM_KV_ROPE_SCALING_YARN_LOG_MUL, hparams.rope_yarn_log_mul, 0.0f)) {
1640+
// [TAG_DEEPSEEK2_YARN_LOG_MUL_FIX]
1641+
// cancel the factor from the convert script
1642+
hparams.rope_yarn_log_mul /= 0.1f;
1643+
}
16391644

16401645
// (optional) temperature tuning - used by mistral-large
16411646
ml.get_key(LLM_KV_ATTENTION_TEMPERATURE_SCALE, hparams.f_attn_temp_scale, false);
@@ -2267,9 +2272,9 @@ void llama_model::load_hparams(llama_model_loader & ml) {
22672272
ml.get_key(LLM_KV_ATTENTION_LAYERNORM_RMS_EPS, hparams.f_norm_rms_eps);
22682273
ml.get_key(LLM_KV_ATTENTION_TEMPERATURE_SCALE, hparams.f_attn_temp_scale, false);
22692274

2270-
ml.get_key(LLM_KV_ROPE_SCALING_YARN_BETA_FAST, hparams.yarn_beta_fast, false);
2271-
ml.get_key(LLM_KV_ROPE_SCALING_YARN_BETA_SLOW, hparams.yarn_beta_slow, false);
2272-
ml.get_key(LLM_KV_ROPE_SCALING_YARN_LOG_MUL, hparams.rope_yarn_log_mul, false);
2275+
ml.get_key(LLM_KV_ROPE_SCALING_YARN_BETA_FAST, hparams.yarn_beta_fast, false);
2276+
ml.get_key(LLM_KV_ROPE_SCALING_YARN_BETA_SLOW, hparams.yarn_beta_slow, false);
2277+
ml.get_key(LLM_KV_ROPE_SCALING_YARN_LOG_MUL, hparams.rope_yarn_log_mul, 0.0f);
22732278

22742279
// TODO: maybe add n_attn_temp_floor_scale as a separate KV?
22752280
if (hparams.f_attn_temp_scale != 0.0f) {
@@ -2279,18 +2284,6 @@ void llama_model::load_hparams(llama_model_loader & ml) {
22792284
}
22802285
}
22812286

2282-
// TODO: this seems to be correct with the case of mscale == mscale_all_dims == 1.0f
2283-
// but may need further verification with other values
2284-
if (hparams.rope_yarn_log_mul != 0.0f) {
2285-
float factor = 1.0f / hparams.rope_freq_scale_train;
2286-
float mscale = 1.0f;
2287-
float mscale_all_dims = hparams.rope_yarn_log_mul;
2288-
static auto get_mscale = [](float scale, float mscale) {
2289-
return scale <= 1.0f ? 1.0f : (0.1f * mscale * logf(scale) + 1.0f);
2290-
};
2291-
hparams.yarn_attn_factor = get_mscale(factor, mscale) / get_mscale(factor, mscale_all_dims);
2292-
}
2293-
22942287
switch (hparams.n_layer) {
22952288
case 26: type = LLM_TYPE_3B; break;
22962289
case 34: type = LLM_TYPE_8B; break;
@@ -2301,6 +2294,32 @@ void llama_model::load_hparams(llama_model_loader & ml) {
23012294
default: throw std::runtime_error("unsupported model architecture");
23022295
}
23032296

2297+
// ref: https://github.com/huggingface/transformers/blob/6d00f6b0a5679c36510f203e4226e36f517c3032/src/transformers/modeling_rope_utils.py#L336-L348
2298+
if (hparams.rope_yarn_log_mul != 0.0f) {
2299+
const float factor = 1.0f / hparams.rope_freq_scale_train;
2300+
2301+
// note: here we assume `mscale == 1.0f`
2302+
// TODO: start reading the actual value of mscale and handle the case where it is not 1.0f
2303+
float mscale = 1.0f;
2304+
const float mscale_all_dims = hparams.rope_yarn_log_mul;
2305+
2306+
// [TAG_DEEPSEEK2_YARN_LOG_MUL_FIX]
2307+
// special-case DEEPSEEK v2:
2308+
// https://huggingface.co/deepseek-ai/DeepSeek-V2-Lite-Chat/blob/main/config.json#L42-L43
2309+
if (arch == LLM_ARCH_DEEPSEEK2 && mscale_all_dims != 1.0f) {
2310+
mscale = mscale_all_dims;
2311+
}
2312+
2313+
static auto get_mscale = [](float scale, float mscale) {
2314+
return scale <= 1.0f ? 1.0f : (0.1f * mscale * logf(scale) + 1.0f);
2315+
};
2316+
2317+
hparams.yarn_attn_factor = get_mscale(factor, mscale) / get_mscale(factor, mscale_all_dims);
2318+
2319+
LLAMA_LOG_WARN("%s: setting new yarn_attn_factor = %.4f (mscale == %.1f, mscale_all_dim = %.1f)\n",
2320+
__func__, hparams.yarn_attn_factor, mscale, mscale_all_dims);
2321+
}
2322+
23042323
pimpl->n_bytes = ml.n_bytes;
23052324

23062325
pimpl->desc_str = arch_name() + " " + type_name() + " " + ml.ftype_name();
@@ -6806,6 +6825,7 @@ void llama_model::print_info() const {
68066825
LLAMA_LOG_INFO("%s: freq_base_train = %.1f\n", __func__, hparams.rope_freq_base_train);
68076826
LLAMA_LOG_INFO("%s: freq_scale_train = %g\n", __func__, hparams.rope_freq_scale_train);
68086827
LLAMA_LOG_INFO("%s: n_ctx_orig_yarn = %u\n", __func__, hparams.n_ctx_orig_yarn);
6828+
LLAMA_LOG_INFO("%s: rope_yarn_log_mul= %.4f\n", __func__, hparams.rope_yarn_log_mul);
68096829
LLAMA_LOG_INFO("%s: rope_finetuned = %s\n", __func__, hparams.rope_finetuned ? "yes" : "unknown");
68106830
// MRoPE (Multi-axis Rotary Position Embedding) sections
68116831
if (const auto & s = hparams.rope_sections; s[0] || s[1] || s[2] || s[3]) {
@@ -6869,7 +6889,6 @@ void llama_model::print_info() const {
68696889
LLAMA_LOG_INFO("%s: expert_weights_scale = %.1f\n", __func__, hparams.expert_weights_scale);
68706890
LLAMA_LOG_INFO("%s: expert_weights_norm = %d\n", __func__, hparams.expert_weights_norm);
68716891
LLAMA_LOG_INFO("%s: expert_gating_func = %s\n", __func__, llama_expert_gating_func_name((llama_expert_gating_func_type) hparams.expert_gating_func));
6872-
LLAMA_LOG_INFO("%s: rope_yarn_log_mul = %.4f\n", __func__, hparams.rope_yarn_log_mul);
68736892
}
68746893

68756894
if (arch == LLM_ARCH_QWEN2MOE) {

src/models/deepseek2.cpp

Lines changed: 9 additions & 5 deletions
Original file line numberDiff line numberDiff line change
@@ -1,7 +1,5 @@
11
#include "models.h"
22

3-
4-
53
llm_build_deepseek2::llm_build_deepseek2(const llama_model & model, const llm_graph_params & params) :
64
llm_graph_context(params) {
75
// lite variants include DeepSeek-V2-Lite, GigaChat3-10B-A1.8B
@@ -20,9 +18,15 @@ llm_build_deepseek2::llm_build_deepseek2(const llama_model & model, const llm_gr
2018

2119
// We have to pre-scale kq_scale and attn_factor to make the YaRN RoPE work correctly.
2220
// See https://github.com/ggerganov/llama.cpp/discussions/7416 for detailed explanation.
23-
const float mscale = attn_factor * (1.0f + hparams.rope_yarn_log_mul * logf(1.0f / freq_scale));
24-
const float kq_scale = 1.0f * mscale * mscale / sqrtf(float(n_embd_head_k));
25-
const float attn_factor = 1.0f / (1.0f + 0.1f * logf(1.0f / freq_scale));
21+
// And also: https://github.com/ggml-org/llama.cpp/pull/17945 [TAG_DEEPSEEK2_YARN_LOG_MUL_FIX]
22+
23+
// first cancel the adjustment from llama_hparams::yarn_attn_factor_adjust to get the original attn_factor
24+
GGML_ASSERT(ext_factor >= 0.0f);
25+
const float attn_factor_org = attn_factor * (1.0f + 0.1f * logf(1.0f / freq_scale));
26+
27+
// use the original attn_factor to pre-scale the kq_scale
28+
const float mscale = attn_factor_org * (1.0f + 0.1f * hparams.rope_yarn_log_mul * logf(1.0f / freq_scale));
29+
const float kq_scale = 1.0f * mscale * mscale / sqrtf(float(n_embd_head_k));
2630

2731
ggml_tensor * cur;
2832
ggml_tensor * inpL;

0 commit comments

Comments
 (0)