diff --git a/.gitignore b/.gitignore index 9b69e2eb4c..7cce05dd99 100644 --- a/.gitignore +++ b/.gitignore @@ -7,4 +7,6 @@ dist .vscode tmp/ requirements-musa.txt -logs/ \ No newline at end of file +logs/ + +benchmark/ diff --git a/lightllm/common/basemodel/attention/fa3/fp.py b/lightllm/common/basemodel/attention/fa3/fp.py index 952bb39d91..a7395faebf 100644 --- a/lightllm/common/basemodel/attention/fa3/fp.py +++ b/lightllm/common/basemodel/attention/fa3/fp.py @@ -7,6 +7,7 @@ from lightllm.utils.envs_utils import get_env_start_args from lightllm.common.basemodel.triton_kernel.fa3_utils import page_table_copy from lightllm.common.basemodel.triton_kernel.gen_prefill_params import gen_cumsum_pad0_tensor +from lightllm.common.basemodel.batch_objs import is_mtp_verify_decode as is_mtp_verify_decode_fn class Fa3AttBackend(BaseAttBackend): @@ -125,8 +126,9 @@ class Fa3DecodeAttState(BaseDecodeAttState): def init_state(self): self.backend: Fa3AttBackend = self.backend args_mtp_step = get_env_start_args().mtp_step + is_mtp_verify_decode = is_mtp_verify_decode_fn(args_mtp_step, self.infer_state.b_num_accepted_tokens) - if args_mtp_step > 0: + if is_mtp_verify_decode: # 修正 mtp 在 fa3 下的输入。 mtp_size = args_mtp_step + 1 b_q_seq_len = torch.full( @@ -143,8 +145,9 @@ def init_state(self): self.cu_seqlens_q = self.infer_state.b1_cu_q_seq_len.int() self.cu_seqlens_k = self.infer_state.b1_cu_kv_seq_len.int() - att_batch_size = self.infer_state.batch_size // (args_mtp_step + 1) - assert self.infer_state.batch_size % (args_mtp_step + 1) == 0 + mtp_size = args_mtp_step + 1 if is_mtp_verify_decode else 1 + att_batch_size = self.infer_state.batch_size // mtp_size + assert self.infer_state.batch_size % mtp_size == 0 model = self.backend.model # 可以使用 cuda graph的时候从 buffer中申请 @@ -163,7 +166,7 @@ def init_state(self): device=self.infer_state.input_ids.device, ) - if args_mtp_step > 0: + if is_mtp_verify_decode: page_table_copy( page_table=self.page_table[:, : self.infer_state.max_kv_seq_len], req_to_token_indexs=model.req_manager.req_to_token_indexs, diff --git a/lightllm/common/basemodel/attention/fa3/fp8.py b/lightllm/common/basemodel/attention/fa3/fp8.py index acbb1315fe..5f275a32a9 100644 --- a/lightllm/common/basemodel/attention/fa3/fp8.py +++ b/lightllm/common/basemodel/attention/fa3/fp8.py @@ -3,7 +3,6 @@ from ..base_att import AttControl from typing import Optional, TYPE_CHECKING from lightllm.utils.sgl_utils import flash_attn_with_kvcache -from lightllm.utils.envs_utils import get_env_start_args from lightllm.common.basemodel.triton_kernel.quantization.q_per_head_fp8_quant import q_per_head_fp8_quant from lightllm.utils.vllm_utils import HAS_VLLM, vllm_ops from typing import Union @@ -116,12 +115,7 @@ def init_state(self): super().init_state() self.backend: Fp8Fa3AttBackend = self.backend - args_mtp_step = get_env_start_args().mtp_step - att_batch_size = self.infer_state.batch_size // (args_mtp_step + 1) - assert self.infer_state.batch_size % (args_mtp_step + 1) == 0 - - device = self.infer_state.input_ids.device - batch_size = att_batch_size + batch_size = self.b_att_seq_len.shape[0] mem_manager = self.backend.model.mem_manager offline_scales: torch.Tensor = mem_manager.scales @@ -180,11 +174,11 @@ def _fp8_decode_att( k_cache=cache_k, v_cache=cache_v, page_table=self.page_table, - cache_seqlens=self.infer_state.b_seq_len, + cache_seqlens=self.b_att_seq_len, cu_seqlens_q=self.cu_seqlens_q, cu_seqlens_k_new=self.cu_seqlens_k, max_seqlen_q=self.decode_max_q_seq_len, - causal=False, + causal=True, window_size=(-1, -1), softcap=0.0, q_descale=q_scale.view(self.infer_state.batch_size, k_head_num), diff --git a/lightllm/common/basemodel/attention/fa3/mla.py b/lightllm/common/basemodel/attention/fa3/mla.py index 9a10457b12..982bd117c3 100644 --- a/lightllm/common/basemodel/attention/fa3/mla.py +++ b/lightllm/common/basemodel/attention/fa3/mla.py @@ -8,6 +8,7 @@ from lightllm.common.basemodel.triton_kernel.fa3_utils import page_table_copy from lightllm.common.basemodel.triton_kernel.gen_prefill_params import gen_cumsum_pad0_tensor from lightllm.utils.sgl_utils import flash_attn_varlen_func +from lightllm.common.basemodel.batch_objs import is_mtp_verify_decode as is_mtp_verify_decode_fn class MlaFa3AttBackend(BaseAttBackend): @@ -108,8 +109,9 @@ class MlaFa3DecodeAttState(BaseDecodeAttState): def init_state(self): self.backend: MlaFa3AttBackend = self.backend args_mtp_step = get_env_start_args().mtp_step + is_mtp_verify_decode = is_mtp_verify_decode_fn(args_mtp_step, self.infer_state.b_num_accepted_tokens) - if args_mtp_step > 0: + if is_mtp_verify_decode: # 修正 mtp 在 fa3 下的输入。 mtp_size = args_mtp_step + 1 b_q_seq_len = torch.full( @@ -126,8 +128,9 @@ def init_state(self): self.cu_seqlens_q = self.infer_state.b1_cu_q_seq_len.int() self.cu_seqlens_k = self.infer_state.b1_cu_kv_seq_len.int() - att_batch_size = self.infer_state.batch_size // (args_mtp_step + 1) - assert self.infer_state.batch_size % (args_mtp_step + 1) == 0 + mtp_size = args_mtp_step + 1 if is_mtp_verify_decode else 1 + att_batch_size = self.infer_state.batch_size // mtp_size + assert self.infer_state.batch_size % mtp_size == 0 model = self.backend.model # 可以使用 cuda graph的时候从 buffer中申请 @@ -146,7 +149,7 @@ def init_state(self): device=self.infer_state.input_ids.device, ) - if args_mtp_step > 0: + if is_mtp_verify_decode: page_table_copy( page_table=self.page_table[:, : self.infer_state.max_kv_seq_len], req_to_token_indexs=model.req_manager.req_to_token_indexs, diff --git a/lightllm/common/basemodel/basemodel.py b/lightllm/common/basemodel/basemodel.py index 473dcbafda..47c13adcb7 100755 --- a/lightllm/common/basemodel/basemodel.py +++ b/lightllm/common/basemodel/basemodel.py @@ -4,6 +4,7 @@ import gc import copy import json +import math import torch import torch.nn.functional as F import triton @@ -17,20 +18,33 @@ from lightllm.common.req_manager import ReqManager from lightllm.common.infer_utils import init_req_to_token_indexes from lightllm.common.build_utils import repair_config -from lightllm.common.basemodel.triton_kernel.copy_kv_index_to_req import copy_kv_index_to_req +from lightllm.common.basemodel.triton_kernel.copy_kv_index_to_req import ( + copy_kv_index_to_req, +) from lightllm.common.basemodel.layer_infer.cache_tensor_manager import g_cache_manager from lightllm.common.basemodel.cuda_graph import CudaGraph from lightllm.common.basemodel.prefill_cuda_graph import PrefillCudaGraph from lightllm.common.quantization import Quantcfg -from lightllm.common.basemodel.triton_kernel.gather_token_id import gather_token, gather_token_prefill_decode_mixed +from lightllm.common.basemodel.triton_kernel.gather_token_id import ( + gather_token, + gather_token_prefill_decode_mixed, +) from lightllm.utils.log_utils import init_logger from lightllm.utils.dist_utils import get_dp_world_size -from lightllm.utils.envs_utils import get_env_start_args, get_llm_data_type, get_added_mtp_kv_layer_num +from lightllm.utils.envs_utils import ( + get_env_start_args, + get_llm_data_type, + get_added_mtp_kv_layer_num, +) from lightllm.distributed.communication_op import dist_group_manager from lightllm.common.basemodel.batch_objs import ModelInput, ModelOutput +from lightllm.common.basemodel.batch_objs import is_mtp_verify_decode as is_mtp_verify_decode_fn from lightllm.common.triton_utils.autotuner import AutotuneLevel from lightllm.utils.custom_kernel_utis import pad2dim_tensor_to_new_batch -from lightllm.utils.envs_utils import set_model_init_status, enable_diverse_mode_gqa_decode_fast_kernel +from lightllm.utils.envs_utils import ( + set_model_init_status, + enable_diverse_mode_gqa_decode_fast_kernel, +) from lightllm.common.triton_utils.autotuner import Autotuner from lightllm.utils.infer_utils import post_empty_cache from .attention import get_prefill_att_backend_class, get_decode_att_backend_class @@ -321,6 +335,7 @@ def _create_inferstate(self, model_input: ModelInput, microbatch_index: int = 0) infer_state.b_req_idx = model_input.b_req_idx infer_state.b_seq_len = model_input.b_seq_len infer_state.b_mtp_index = model_input.b_mtp_index + infer_state.b_num_accepted_tokens = model_input.b_num_accepted_tokens if model_input.is_prefill: if model_input.b_ready_cache_len is not None: infer_state.b_ready_cache_len = model_input.b_ready_cache_len @@ -358,6 +373,18 @@ def _create_inferstate(self, model_input: ModelInput, microbatch_index: int = 0) return infer_state + def _get_decode_padding_unit(self, model_input: ModelInput) -> int: + padding_unit = self.tp_world_size_ if self.args.enable_tpsp_mix_mode else 1 + if (not model_input.is_prefill) and is_mtp_verify_decode_fn( + self.args.mtp_step, model_input.b_num_accepted_tokens + ): + padding_unit = math.lcm(padding_unit, self.args.mtp_step + 1) + return padding_unit + + def _get_decode_infer_batch_size(self, model_input: ModelInput) -> int: + padding_unit = self._get_decode_padding_unit(model_input) + return triton.cdiv(model_input.batch_size, padding_unit) * padding_unit + def _create_padded_decode_model_input(self, model_input: ModelInput, new_batch_size: int): if model_input.batch_size == new_batch_size: return model_input @@ -367,22 +394,111 @@ def _create_padded_decode_model_input(self, model_input: ModelInput, new_batch_s padded_batch_size = new_batch_size - model_input.batch_size new_model_input = copy.copy(model_input) new_model_input.batch_size = new_batch_size - new_model_input.total_token_num += padded_batch_size * 2 - new_model_input.max_kv_seq_len = max(2, model_input.max_kv_seq_len) - new_model_input.input_ids = F.pad(new_model_input.input_ids, (0, padded_batch_size), mode="constant", value=1) - new_model_input.b_req_idx = F.pad( - new_model_input.b_req_idx, (0, padded_batch_size), mode="constant", value=self.req_manager.HOLD_REQUEST_ID - ) - new_model_input.b_mtp_index = F.pad( - new_model_input.b_mtp_index, (0, padded_batch_size), mode="constant", value=0 - ) - new_model_input.b_seq_len = F.pad(new_model_input.b_seq_len, (0, padded_batch_size), mode="constant", value=2) - new_model_input.mem_indexes = F.pad( - new_model_input.mem_indexes, - (0, padded_batch_size), - mode="constant", - value=self.mem_manager.HOLD_TOKEN_MEMINDEX, + + is_mtp_verify_decode = (not model_input.is_prefill) and is_mtp_verify_decode_fn( + self.args.mtp_step, model_input.b_num_accepted_tokens ) + if is_mtp_verify_decode: + mtp_size = self.args.mtp_step + 1 + assert model_input.batch_size % mtp_size == 0 + assert new_batch_size % mtp_size == 0 + assert padded_batch_size % mtp_size == 0 + padded_req_num = padded_batch_size // mtp_size + + pad_mtp_index = torch.arange( + mtp_size, + dtype=new_model_input.b_mtp_index.dtype, + device=new_model_input.b_mtp_index.device, + ).repeat(padded_req_num) + pad_seq_len = torch.arange( + 2, + mtp_size + 2, + dtype=new_model_input.b_seq_len.dtype, + device=new_model_input.b_seq_len.device, + ).repeat(padded_req_num) + new_model_input.total_token_num += padded_req_num * (mtp_size * (mtp_size + 3) // 2) + new_model_input.max_kv_seq_len = max(mtp_size + 1, model_input.max_kv_seq_len) + new_model_input.input_ids = torch.cat( + ( + new_model_input.input_ids, + torch.ones( + padded_batch_size, + dtype=new_model_input.input_ids.dtype, + device=new_model_input.input_ids.device, + ), + ), + dim=0, + ) + new_model_input.b_req_idx = torch.cat( + ( + new_model_input.b_req_idx, + torch.full( + (padded_batch_size,), + self.req_manager.HOLD_REQUEST_ID, + dtype=new_model_input.b_req_idx.dtype, + device=new_model_input.b_req_idx.device, + ), + ), + dim=0, + ) + new_model_input.b_mtp_index = torch.cat((new_model_input.b_mtp_index, pad_mtp_index), dim=0) + new_model_input.b_seq_len = torch.cat((new_model_input.b_seq_len, pad_seq_len), dim=0) + new_model_input.mem_indexes = torch.cat( + ( + new_model_input.mem_indexes, + torch.full( + (padded_batch_size,), + self.mem_manager.HOLD_TOKEN_MEMINDEX, + dtype=new_model_input.mem_indexes.dtype, + device=new_model_input.mem_indexes.device, + ), + ), + dim=0, + ) + new_model_input.b_num_accepted_tokens = torch.cat( + ( + new_model_input.b_num_accepted_tokens, + torch.ones( + padded_req_num, + dtype=new_model_input.b_num_accepted_tokens.dtype, + device=new_model_input.b_num_accepted_tokens.device, + ), + ), + dim=0, + ) + else: + new_model_input.total_token_num += padded_batch_size * 2 + new_model_input.max_kv_seq_len = max(2, model_input.max_kv_seq_len) + new_model_input.input_ids = F.pad( + new_model_input.input_ids, + (0, padded_batch_size), + mode="constant", + value=1, + ) + new_model_input.b_req_idx = F.pad( + new_model_input.b_req_idx, + (0, padded_batch_size), + mode="constant", + value=self.req_manager.HOLD_REQUEST_ID, + ) + new_model_input.b_mtp_index = F.pad( + new_model_input.b_mtp_index, + (0, padded_batch_size), + mode="constant", + value=0, + ) + new_model_input.b_seq_len = F.pad( + new_model_input.b_seq_len, + (0, padded_batch_size), + mode="constant", + value=2, + ) + new_model_input.mem_indexes = F.pad( + new_model_input.mem_indexes, + (0, padded_batch_size), + mode="constant", + value=self.mem_manager.HOLD_TOKEN_MEMINDEX, + ) new_model_input.multimodal_params = new_model_input.multimodal_params + [ {"images": [], "audios": []} for _ in range(padded_batch_size) ] @@ -390,11 +506,17 @@ def _create_padded_decode_model_input(self, model_input: ModelInput, new_batch_s if enable_diverse_mode_gqa_decode_fast_kernel(): if new_model_input.b_shared_seq_len is not None: new_model_input.b_shared_seq_len = F.pad( - new_model_input.b_shared_seq_len, (0, padded_batch_size), mode="constant", value=0 + new_model_input.b_shared_seq_len, + (0, padded_batch_size), + mode="constant", + value=0, ) if new_model_input.b_mark_shared_group is not None: new_model_input.b_mark_shared_group = F.pad( - new_model_input.b_mark_shared_group, (0, padded_batch_size), mode="constant", value=1 + new_model_input.b_mark_shared_group, + (0, padded_batch_size), + mode="constant", + value=1, ) # 特殊模型,特殊模式的特殊变量的特殊 padding @@ -429,7 +551,10 @@ def _create_padded_prefill_model_input(self, model_input: ModelInput, new_handle value=self.mem_manager.HOLD_TOKEN_MEMINDEX, ) new_model_input.b_req_idx = F.pad( - new_model_input.b_req_idx, (0, 1), mode="constant", value=self.req_manager.HOLD_REQUEST_ID + new_model_input.b_req_idx, + (0, 1), + mode="constant", + value=self.req_manager.HOLD_REQUEST_ID, ) new_model_input.b_mtp_index = F.pad(new_model_input.b_mtp_index, (0, 1), mode="constant", value=0) new_model_input.b_seq_len = F.pad(new_model_input.b_seq_len, (0, 1), mode="constant", value=padded_token_num) @@ -469,7 +594,10 @@ def _create_unpad_decode_model_output(self, model_output: ModelOutput, origin_ba return new_model_output def _create_unpad_prefill_model_output( - self, padded_model_output: ModelOutput, origin_handle_token_num: int, origin_batch_size: int + self, + padded_model_output: ModelOutput, + origin_handle_token_num: int, + origin_batch_size: int, ): if self.return_all_prompt_logics: new_model_output = copy.copy(padded_model_output) @@ -555,15 +683,16 @@ def _decode( ) origin_batch_size = model_input.batch_size - if self.args.enable_tpsp_mix_mode: - infer_batch_size = triton.cdiv(model_input.batch_size, self.tp_world_size_) * self.tp_world_size_ - else: - infer_batch_size = model_input.batch_size + infer_batch_size = self._get_decode_infer_batch_size(model_input) + is_mtp_verify_decode = is_mtp_verify_decode_fn(self.args.mtp_step, model_input.b_num_accepted_tokens) if self.graph is not None and self.graph.can_run( batch_size=infer_batch_size, max_len_in_batch=model_input.max_kv_seq_len ): - infer_batch_size = self.graph.find_closest_graph_batch_size(batch_size=infer_batch_size) + infer_batch_size = self.graph.find_closest_graph_batch_size( + batch_size=infer_batch_size, + is_mtp_verify_decode=is_mtp_verify_decode, + ) model_input = self._create_padded_decode_model_input( model_input=model_input, new_batch_size=infer_batch_size ) @@ -577,7 +706,7 @@ def _decode( infer_state.init_some_extra_state(self) infer_state.init_att_state() - if self.graph.need_capture(infer_batch_size): + if self.graph.need_capture(infer_batch_size, is_mtp_verify_decode=is_mtp_verify_decode): infer_state.is_cuda_graph = True model_output: ModelOutput = self.graph.capture_decode(self._token_forward, infer_state) else: @@ -604,7 +733,6 @@ def _decode( @final def _context_forward(self, infer_state: InferStateInfo): - input_embs = self.pre_infer.context_forward(infer_state.input_ids, infer_state, self.pre_post_weight) if self.args.enable_dp_prefill_balance: assert not self.args.enable_prefill_cudagraph, "not support now" @@ -810,10 +938,14 @@ def microbatch_overlap_decode(self, model_input0: ModelInput, model_input1: Mode origin_batch_size = model_input0.batch_size max_len_in_batch = max(model_input0.max_kv_seq_len, model_input1.max_kv_seq_len) - infer_batch_size = triton.cdiv(origin_batch_size, self.tp_world_size_) * self.tp_world_size_ + infer_batch_size = self._get_decode_infer_batch_size(model_input0) + is_mtp_verify_decode = is_mtp_verify_decode_fn(self.args.mtp_step, model_input0.b_num_accepted_tokens) if self.graph is not None and self.graph.can_run(infer_batch_size, max_len_in_batch): - infer_batch_size = self.graph.find_closest_graph_batch_size(infer_batch_size) + infer_batch_size = self.graph.find_closest_graph_batch_size( + infer_batch_size, + is_mtp_verify_decode=is_mtp_verify_decode, + ) # TODO 如果支持动态步数的 mtp,在不同的mtp步上,model_input0 和 model_input1 的内部batch size可能不 # 一致,需要按照较高 batch size 进行graph的寻找,同时,进行有效的恢复。 padded_model_input0 = self._create_padded_decode_model_input(model_input0, infer_batch_size) @@ -838,7 +970,7 @@ def microbatch_overlap_decode(self, model_input0: ModelInput, model_input1: Mode infer_state1.init_some_extra_state(self) infer_state1.init_att_state() - if self.graph.need_capture(infer_batch_size): + if self.graph.need_capture(infer_batch_size, is_mtp_verify_decode=is_mtp_verify_decode): infer_state0.is_cuda_graph = True infer_state1.is_cuda_graph = True @@ -890,7 +1022,11 @@ def _overlap_tpsp_context_forward(self, infer_state: InferStateInfo, infer_state g_cache_manager.cache_env_in() input_embs, input_embs1 = self.pre_infer.overlap_tpsp_context_forward( - infer_state.input_ids, infer_state1.input_ids, infer_state, infer_state1, self.pre_post_weight + infer_state.input_ids, + infer_state1.input_ids, + infer_state, + infer_state1, + self.pre_post_weight, ) # 决定是否进行 dp balance 优化,可以提升dp > 1 时的 prefill 效率。 @@ -906,7 +1042,11 @@ def _overlap_tpsp_context_forward(self, infer_state: InferStateInfo, infer_state for i in range(self.layers_num): input_embs, input_embs1 = self.layers_infer[i].overlap_tpsp_context_forward( - input_embs, input_embs1, infer_state, infer_state1, self.trans_layers_weight[i] + input_embs, + input_embs1, + infer_state, + infer_state1, + self.trans_layers_weight[i], ) # 折叠模式调用完infer_state 和 infer_state1 上的hook函数后,input_embs 和 input_embs1 才具备正确的运算数据。 @@ -920,7 +1060,11 @@ def _overlap_tpsp_context_forward(self, infer_state: InferStateInfo, infer_state last_input_embs1 = infer_state1._all_to_all_unbalance_get(data=last_input_embs1) predict_logits, predict_logits1 = self.post_infer.overlap_tpsp_token_forward( - last_input_embs, last_input_embs1, infer_state, infer_state1, self.pre_post_weight + last_input_embs, + last_input_embs1, + infer_state, + infer_state1, + self.pre_post_weight, ) g_cache_manager.cache_env_out() @@ -941,14 +1085,22 @@ def _overlap_tpsp_context_forward(self, infer_state: InferStateInfo, infer_state @final def _overlap_tpsp_token_forward(self, infer_state: InferStateInfo, infer_state1: InferStateInfo): input_embs, input_embs1 = self.pre_infer.overlap_tpsp_token_forward( - infer_state.input_ids, infer_state1.input_ids, infer_state, infer_state1, self.pre_post_weight + infer_state.input_ids, + infer_state1.input_ids, + infer_state, + infer_state1, + self.pre_post_weight, ) input_embs = self.pre_infer._tpsp_sp_split(input=input_embs, infer_state=infer_state) input_embs1 = self.pre_infer._tpsp_sp_split(input=input_embs1, infer_state=infer_state1) for i in range(self.layers_num): input_embs, input_embs1 = self.layers_infer[i].overlap_tpsp_token_forward( - input_embs, input_embs1, infer_state, infer_state1, self.trans_layers_weight[i] + input_embs, + input_embs1, + infer_state, + infer_state1, + self.trans_layers_weight[i], ) # 折叠模式调用完infer_state 上的hook函数后,input_embs 和 input_embs 才具备正确的运算数据。 @@ -959,7 +1111,11 @@ def _overlap_tpsp_token_forward(self, infer_state: InferStateInfo, infer_state1: last_input_embs1 = self.post_infer._tpsp_allgather(input=input_embs1, infer_state=infer_state1) predict_logits, predict_logits1 = self.post_infer.overlap_tpsp_token_forward( - last_input_embs, last_input_embs1, infer_state, infer_state1, self.pre_post_weight + last_input_embs, + last_input_embs1, + infer_state, + infer_state1, + self.pre_post_weight, ) model_output = ModelOutput(logits=predict_logits.contiguous()) @@ -1066,7 +1222,12 @@ def _autotune_warmup(self): rand_gen = torch.Generator(device="cuda") rand_gen.manual_seed(input_len) dummy_input_ids = torch.randint( - 0, 10000, (input_len,), dtype=torch.int32, device="cuda", generator=rand_gen + 0, + 10000, + (input_len,), + dtype=torch.int32, + device="cuda", + generator=rand_gen, ) b_req_idx = torch.tensor([self.req_manager.alloc()], dtype=torch.int32, device="cuda") mem_indexes = self.mem_manager.alloc(len(dummy_input_ids)).cuda() @@ -1130,10 +1291,14 @@ def _init_padded_req(self): batch_size = 1 dummy_input_ids = torch.ones((batch_size,), dtype=torch.int32, device="cuda") b_req_idx = torch.tensor( - [self.req_manager.HOLD_REQUEST_ID for _ in range(batch_size)], dtype=torch.int32, device="cuda" + [self.req_manager.HOLD_REQUEST_ID for _ in range(batch_size)], + dtype=torch.int32, + device="cuda", ) mem_indexes = torch.tensor( - [self.mem_manager.HOLD_TOKEN_MEMINDEX for _ in range(batch_size)], dtype=torch.int32, device="cuda" + [self.mem_manager.HOLD_TOKEN_MEMINDEX for _ in range(batch_size)], + dtype=torch.int32, + device="cuda", ) b_seq_len = torch.ones(batch_size, dtype=torch.int32, device="cuda") b_ready_cache_len = torch.zeros(batch_size, dtype=torch.int32, device="cuda") @@ -1177,15 +1342,13 @@ def _init_padded_req(self): def _gen_special_model_input(self, token_num: int): special_model_input = {} - is_mtp_draft_model = ( - "Deepseek3MTPModel" in str(self.__class__) - or "Qwen3MOEMTPModel" in str(self.__class__) - or "MistralMTPModel" in str(self.__class__) - or "Glm4MoeLiteMTPModel" in str(self.__class__) - ) + is_mtp_draft_model = getattr(self, "is_mtp_draft_model", False) if is_mtp_draft_model: special_model_input["mtp_draft_input_hiddens"] = torch.randn( - token_num, self.config["hidden_size"], dtype=self.data_type, device="cuda" + token_num, + self.config["hidden_size"], + dtype=self.data_type, + device="cuda", ) else: special_model_input["mtp_draft_input_hiddens"] = None diff --git a/lightllm/common/basemodel/batch_objs.py b/lightllm/common/basemodel/batch_objs.py index 1795ff9a82..6104022733 100644 --- a/lightllm/common/basemodel/batch_objs.py +++ b/lightllm/common/basemodel/batch_objs.py @@ -6,6 +6,13 @@ from lightllm.utils.tensor_utils import tensor_to_no_ref_tensor +def is_mtp_verify_decode(mtp_step: int, b_num_accepted_tokens) -> bool: + """Single source of truth for the MTP verify-decode predicate (#21). + A decode forward is a verify pass iff MTP is enabled and the per-real-request accept tensor is + present — decode_mtp sets it on the main verify and clears it (None) on every draft forward.""" + return mtp_step > 0 and b_num_accepted_tokens is not None + + @dataclass class ModelInput: # 通用变量 @@ -53,6 +60,8 @@ class ModelInput: # 的 draft 模型的输入 mtp_draft_input_hiddens: Optional[torch.Tensor] = None + b_num_accepted_tokens: Optional[torch.Tensor] = None + def to_cuda(self): if self.input_ids is not None: self.input_ids = self.input_ids.cuda(non_blocking=True) @@ -66,6 +75,8 @@ def to_cuda(self): self.b_req_idx = self.b_req_idx.cuda(non_blocking=True) self.b_seq_len = self.b_seq_len.cuda(non_blocking=True) self.b_mtp_index = self.b_mtp_index.cuda(non_blocking=True) + if self.b_num_accepted_tokens is not None: + self.b_num_accepted_tokens = self.b_num_accepted_tokens.cuda(non_blocking=True) if self.b_ready_cache_len is not None: self.b_ready_cache_len = self.b_ready_cache_len.cuda(non_blocking=True) if self.b_prefill_start_loc is not None: diff --git a/lightllm/common/basemodel/cuda_graph.py b/lightllm/common/basemodel/cuda_graph.py index 782150661e..d20de7afb8 100644 --- a/lightllm/common/basemodel/cuda_graph.py +++ b/lightllm/common/basemodel/cuda_graph.py @@ -2,12 +2,14 @@ import torch import copy import bisect +import math import triton from typing import Optional from lightllm.utils.log_utils import init_logger from lightllm.utils.envs_utils import get_env_start_args from lightllm.distributed import dist_group_manager from lightllm.common.basemodel.batch_objs import ModelInput, ModelOutput +from lightllm.common.basemodel.batch_objs import is_mtp_verify_decode as is_mtp_verify_decode_fn from .infer_struct import InferStateInfo @@ -27,48 +29,135 @@ def __init__(self, max_batch_size=8, max_len_in_batch=8192, tp_world_size: int = self.graph_max_len_in_batch = max_len_in_batch self.enable_decode_microbatch_overlap = self.args.enable_decode_microbatch_overlap + self.normal_cuda_graph_batch_sizes = self._build_cuda_graph_batch_sizes(batch_size_multiple=1) + if self.mtp_step > 0: + self.mtp_verify_cuda_graph_batch_sizes = self._build_cuda_graph_batch_sizes( + batch_size_multiple=self.mtp_step + 1 + ) + logger.info(f"normal cuda graph batch_sizes: {self.normal_cuda_graph_batch_sizes}") + logger.info(f"mtp verify cuda graph batch_sizes: {self.mtp_verify_cuda_graph_batch_sizes}") + else: + self.mtp_verify_cuda_graph_batch_sizes = self.normal_cuda_graph_batch_sizes + logger.info(f"cuda graph batch_sizes: {self.normal_cuda_graph_batch_sizes}") + + def _build_cuda_graph_batch_sizes(self, batch_size_multiple: int): # gen cuda graph batch_sizes # cuda graph gen for batch size = [1, 2, 3, ..., graph_split_batch_size] - # and [graph_split_batch_size + graph_grow_step_size, - # if the mtp_step is not 0, then the batch_sizes will be multiply of (mtp_step + 1) - - graph_split_batch_size = self.args.graph_split_batch_size * (self.mtp_step + 1) - graph_grow_step_size = self.args.graph_grow_step_size * (self.mtp_step + 1) - - batch_sizes = [i * (self.mtp_step + 1) for i in range(1, self.args.graph_split_batch_size + 1)] - for _batch_size in range(graph_split_batch_size + graph_grow_step_size, max_batch_size, graph_grow_step_size): + # and [graph_split_batch_size + graph_grow_step_size, ...] + graph_split_batch_size = self.args.graph_split_batch_size * batch_size_multiple + graph_grow_step_size = self.args.graph_grow_step_size * batch_size_multiple + + batch_sizes = [i * batch_size_multiple for i in range(1, self.args.graph_split_batch_size + 1)] + for _batch_size in range( + graph_split_batch_size + graph_grow_step_size, + self.max_batch_size, + graph_grow_step_size, + ): batch_sizes.append(_batch_size) - batch_sizes = list(set([e for e in batch_sizes if e < max_batch_size])) - batch_sizes.append(max_batch_size) + batch_sizes = list(set([e for e in batch_sizes if e < self.max_batch_size])) + batch_sizes.append(self.max_batch_size) batch_sizes.sort() if self.args.enable_tpsp_mix_mode: - batch_sizes = [triton.cdiv(e, self.tp_world_size) * self.tp_world_size for e in batch_sizes] + padding_unit = math.lcm(self.tp_world_size, batch_size_multiple) + batch_sizes = [triton.cdiv(e, padding_unit) * padding_unit for e in batch_sizes] batch_sizes = list(set(batch_sizes)) batch_sizes.sort() - self.cuda_graph_batch_sizes = batch_sizes assert batch_sizes[-1] == self.max_batch_size - logger.info(f"cuda graph batch_sizes: {self.cuda_graph_batch_sizes}") + return batch_sizes def can_run(self, batch_size, max_len_in_batch): return batch_size <= self.max_batch_size and max_len_in_batch <= self.graph_max_len_in_batch - def need_capture(self, batch_size): - find_batch_size = self.find_closest_graph_batch_size(batch_size) + def _decode_graph_key(self, infer_state: InferStateInfo): + is_mtp_verify_decode = is_mtp_verify_decode_fn(self.mtp_step, infer_state.b_num_accepted_tokens) + return (infer_state.input_ids.shape[0], is_mtp_verify_decode) + + def need_capture(self, batch_size, is_mtp_verify_decode=False): + find_batch_size = self.find_closest_graph_batch_size(batch_size, is_mtp_verify_decode=is_mtp_verify_decode) if find_batch_size is not None: - return find_batch_size not in self.graph + return (find_batch_size, is_mtp_verify_decode) not in self.graph else: assert False, "dead code" - def find_closest_graph_batch_size(self, batch_size): - index = bisect.bisect_left(self.cuda_graph_batch_sizes, batch_size) - if index < len(self.cuda_graph_batch_sizes): - find_batch_size = self.cuda_graph_batch_sizes[index] + def _get_graph_batch_sizes(self, is_mtp_verify_decode=False): + if is_mtp_verify_decode: + return self.mtp_verify_cuda_graph_batch_sizes + return self.normal_cuda_graph_batch_sizes + + def find_closest_graph_batch_size(self, batch_size, is_mtp_verify_decode=False): + graph_batch_sizes = self._get_graph_batch_sizes(is_mtp_verify_decode=is_mtp_verify_decode) + index = bisect.bisect_left(graph_batch_sizes, batch_size) + if index < len(graph_batch_sizes): + find_batch_size = graph_batch_sizes[index] return find_batch_size else: return None + def _build_warmup_decode_model_input( + self, + model, + batch_size: int, + device: str = "cuda", + is_mtp_verify_decode: Optional[bool] = None, + ) -> ModelInput: + if is_mtp_verify_decode is None: + is_mtp_verify_decode = self.mtp_step > 0 + + mtp_size = self.mtp_step + 1 + input_ids = torch.ones(batch_size, dtype=torch.int32, device=device) + mem_indexes = model.mem_manager.alloc(batch_size).to(device) + b_req_idx = torch.full( + (batch_size,), + fill_value=model.req_manager.HOLD_REQUEST_ID, + dtype=torch.int32, + device=device, + ) + + b_num_accepted_tokens = None + if self.mtp_step > 0 and is_mtp_verify_decode: + assert batch_size % mtp_size == 0, "MTP decode CUDA graph batch size must be a multiple of mtp_step + 1" + real_batch_size = batch_size // mtp_size + b_mtp_index = torch.arange(mtp_size, dtype=torch.int32, device=device).repeat(real_batch_size) + b_seq_len = torch.arange(2, mtp_size + 2, dtype=torch.int32, device=device).repeat(real_batch_size) + b_num_accepted_tokens = torch.ones(real_batch_size, dtype=torch.int32, device=device) + total_token_num = real_batch_size * (mtp_size * (mtp_size + 3) // 2) + else: + seq_len = 2 + total_token_num = batch_size * seq_len + b_mtp_index = torch.zeros(batch_size, dtype=torch.int32, device=device) + b_seq_len = torch.empty(batch_size, dtype=torch.int32, device=device) + b_seq_len.fill_(seq_len) + + return ModelInput( + batch_size=batch_size, + total_token_num=total_token_num, + max_q_seq_len=1, + max_kv_seq_len=self.graph_max_len_in_batch, + input_ids=input_ids, + mem_indexes=mem_indexes, + b_req_idx=b_req_idx, + b_seq_len=b_seq_len, + b_mtp_index=b_mtp_index, + b_num_accepted_tokens=b_num_accepted_tokens, + is_prefill=False, + multimodal_params=[{"images": [], "audios": []} for _ in range(batch_size)], + **model._gen_special_model_input(batch_size), + ) + + def _is_mtp_draft_model(self, model): + return getattr(model, "is_mtp_draft_model", False) + + def _iter_warmup_graph_layouts(self, model): + if self.mtp_step > 0: + if self._is_mtp_draft_model(model): + yield False, self.normal_cuda_graph_batch_sizes + else: + yield True, self.mtp_verify_cuda_graph_batch_sizes + else: + yield False, self.normal_cuda_graph_batch_sizes + def _capture_decode(self, decode_func, infer_state: InferStateInfo): graph_obj = torch.cuda.CUDAGraph() input_ids = infer_state.input_ids @@ -96,7 +185,11 @@ def _capture_decode(self, decode_func, infer_state: InferStateInfo): with torch.cuda.graph(graph_obj, pool=self.mempool): model_output = decode_func(infer_state) - self.graph[batch_size] = (graph_obj, infer_state, model_output) + self.graph[self._decode_graph_key(infer_state)] = ( + graph_obj, + infer_state, + model_output, + ) graph_obj.replay() return model_output @@ -130,7 +223,7 @@ def _capture_decode_overlap( with torch.cuda.graph(graph_obj, pool=self.mempool): model_output, model_output1 = decode_func(infer_state, infer_state1) - self.graph[batch_size] = ( + self.graph[self._decode_graph_key(infer_state)] = ( graph_obj, infer_state, infer_state1, @@ -157,8 +250,7 @@ def capture_decode( return self._capture_decode(decode_func, infer_state) def _replay(self, infer_state: InferStateInfo): - batch_size = infer_state.input_ids.shape[0] - graph_obj, graph_infer_state, graph_output = self.graph[batch_size] + graph_obj, graph_infer_state, graph_output = self.graph[self._decode_graph_key(infer_state)] graph_infer_state.copy_for_cuda_graph(infer_state) graph_obj.replay() return graph_output @@ -168,14 +260,13 @@ def _replay_overlap( infer_state: InferStateInfo, infer_state1: InferStateInfo, ): - batch_size = infer_state.input_ids.shape[0] ( graph_obj, graph_infer_state, graph_infer_state1, graph_model_output, graph_model_output1, - ) = self.graph[batch_size] + ) = self.graph[self._decode_graph_key(infer_state)] graph_infer_state.copy_for_cuda_graph(infer_state) graph_infer_state1.copy_for_cuda_graph(infer_state1) graph_obj.replay() @@ -197,47 +288,23 @@ def warmup(self, model): model: TpPartBaseModel = model # decode cuda graph init - for batch_size in self.cuda_graph_batch_sizes[::-1]: - seq_len = 2 - total_token_num = batch_size * seq_len - max_len_in_batch = self.graph_max_len_in_batch - input_ids = torch.tensor([1 for _ in range(batch_size)], dtype=torch.int32, device="cuda") - mem_indexes = model.mem_manager.alloc(len(input_ids)).cuda() - b_req_idx = torch.tensor( - [model.req_manager.HOLD_REQUEST_ID for _ in range(batch_size)], dtype=torch.int32, device="cuda" - ) - b_seq_len = torch.empty(batch_size, dtype=torch.int32, device="cuda") - b_seq_len.fill_(seq_len) - b_mtp_index = torch.zeros(batch_size, dtype=torch.int32, device="cuda") - - model_input = ModelInput( - batch_size=batch_size, - total_token_num=total_token_num, - max_q_seq_len=1, - max_kv_seq_len=max_len_in_batch, - input_ids=input_ids, - mem_indexes=mem_indexes, - b_req_idx=b_req_idx, - b_seq_len=b_seq_len, - b_mtp_index=b_mtp_index, - is_prefill=False, - multimodal_params=[{"images": [], "audios": []} for _ in range(batch_size)], - **model._gen_special_model_input(batch_size), - ) - model_output: ModelOutput = model.forward(model_input) - del model_output - del input_ids - del mem_indexes - del b_req_idx - del b_seq_len - - model.mem_manager.free_all() - model.req_manager.free_all() - # release local tensors - for var_name, var_value in list(locals().items()): - if isinstance(var_value, torch.Tensor): - del locals()[var_name] - torch.cuda.empty_cache() + for is_mtp_verify_decode, batch_sizes in self._iter_warmup_graph_layouts(model): + for batch_size in batch_sizes[::-1]: + model_input = self._build_warmup_decode_model_input( + model, + batch_size, + is_mtp_verify_decode=is_mtp_verify_decode, + ) + model_output: ModelOutput = model.forward(model_input) + del model_output + + model.mem_manager.free_all() + model.req_manager.free_all() + # release local tensors + for var_name, var_value in list(locals().items()): + if isinstance(var_value, torch.Tensor): + del locals()[var_name] + torch.cuda.empty_cache() logger.info( f"Capture cudagraph success, batch_size <={self.max_batch_size} " @@ -252,56 +319,36 @@ def warmup_overlap(self, model): model: TpPartBaseModel = model - for batch_size in self.cuda_graph_batch_sizes[::-1]: - decode_batches = [] - for micro_batch_index in [0, 1]: - # dummy decoding, capture the cudagraph - seq_len = 2 - total_token_num = batch_size * seq_len - max_len_in_batch = self.graph_max_len_in_batch - input_ids = torch.tensor([1 for _ in range(batch_size)], dtype=torch.int32, device="cuda") - mem_indexes = model.mem_manager.alloc(len(input_ids)).cuda() - b_req_idx = torch.tensor( - [model.req_manager.HOLD_REQUEST_ID for _ in range(batch_size)], dtype=torch.int32, device="cuda" - ) - b_seq_len = torch.empty(batch_size, dtype=torch.int32, device="cuda") - b_seq_len.fill_(seq_len) - b_mtp_index = torch.zeros(batch_size, dtype=torch.int32, device="cuda") - - micro_batch = ModelInput( - is_prefill=False, - batch_size=batch_size, - total_token_num=total_token_num, - max_q_seq_len=1, - max_kv_seq_len=max_len_in_batch, - input_ids=input_ids, - b_mtp_index=b_mtp_index, - mem_indexes=mem_indexes, - b_req_idx=b_req_idx, - b_seq_len=b_seq_len, - multimodal_params=[{"images": [], "audios": []} for _ in range(batch_size)], - **model._gen_special_model_input(batch_size), - ) - decode_batches.append(micro_batch) - del micro_batch + for is_mtp_verify_decode, batch_sizes in self._iter_warmup_graph_layouts(model): + for batch_size in batch_sizes[::-1]: + decode_batches = [] + for micro_batch_index in [0, 1]: + # dummy decoding, capture the cudagraph + micro_batch = self._build_warmup_decode_model_input( + model, + batch_size, + is_mtp_verify_decode=is_mtp_verify_decode, + ) + decode_batches.append(micro_batch) + del micro_batch - for var_name, var_value in list(locals().items()): - if isinstance(var_value, torch.Tensor): - del locals()[var_name] - torch.cuda.empty_cache() + for var_name, var_value in list(locals().items()): + if isinstance(var_value, torch.Tensor): + del locals()[var_name] + torch.cuda.empty_cache() - _, _ = model.microbatch_overlap_decode(decode_batches[0], decode_batches[1]) + _, _ = model.microbatch_overlap_decode(decode_batches[0], decode_batches[1]) - model.mem_manager.free_all() - model.req_manager.free_all() + model.mem_manager.free_all() + model.req_manager.free_all() - del decode_batches + del decode_batches - # release local tensors - for var_name, var_value in list(locals().items()): - if isinstance(var_value, torch.Tensor): - del locals()[var_name] - torch.cuda.empty_cache() + # release local tensors + for var_name, var_value in list(locals().items()): + if isinstance(var_value, torch.Tensor): + del locals()[var_name] + torch.cuda.empty_cache() logger.info( f"Capture overlap cudagraph success, batch_size <={self.max_batch_size} " diff --git a/lightllm/common/basemodel/infer_struct.py b/lightllm/common/basemodel/infer_struct.py index 711484c835..6de15f8910 100755 --- a/lightllm/common/basemodel/infer_struct.py +++ b/lightllm/common/basemodel/infer_struct.py @@ -39,6 +39,8 @@ def __init__(self): self.b_mtp_index: torch.Tensor = None + self.b_num_accepted_tokens: torch.Tensor = None + self.b_seq_len: torch.Tensor = None # max_cache_len 用于 prefill 阶段标识请求中最大 cache的kv 的长度 self.max_cache_len: int = None diff --git a/lightllm/common/basemodel/mtp_verify_extra_state.py b/lightllm/common/basemodel/mtp_verify_extra_state.py new file mode 100644 index 0000000000..e4bc5f4f7a --- /dev/null +++ b/lightllm/common/basemodel/mtp_verify_extra_state.py @@ -0,0 +1,47 @@ +import torch + +from lightllm.utils.envs_utils import get_env_start_args + + +def init_mtp_verify_extra_state(self): + """Shared MTP-verify decode metadata, used by qwen3_5 and qwen3next infer-struct classes (#12). + Call AFTER super().init_some_extra_state(model). `self` is the InferStateInfo instance.""" + self.b_att_seq_len = self.b_seq_len + mtp_step = get_env_start_args().mtp_step + self.b_buffer_idx = self.b_req_idx * (mtp_step + 1) + self.b_mtp_index + # conv buffer is now ONE widened slot per request (indexed by req_idx), + # dropping the *(S+1) + mtp_index addressing used by the SSM block. + self.b_conv_buffer_idx = self.b_req_idx + # MTP verify batch: decode-mode, S+1 expanded, and gated on the + # per-real-request accept tensor that decode_mtp threads in. Gating on + # b_num_accepted_tokens (vs only b_mtp_index, which is set for any decode) + # distinguishes the main-model verify forward from draft/plain decode. + self.is_mtp_verify = ( + (mtp_step > 0) + and (not self.is_prefill) + and (self.b_mtp_index is not None) + and (self.b_num_accepted_tokens is not None) + ) + self.b_gdn_verify_cu_seqlens = None + self.b_ssm_index_rows = None + # b_num_accepted_tokens is threaded onto the infer_state from ModelInput by + # _create_inferstate (mirrors b_mtp_index) BEFORE this runs; nothing to do here. + if self.is_mtp_verify: + step = mtp_step + 1 + n_real = self.b_req_idx.shape[0] // step + self.b_gdn_verify_cu_seqlens = torch.arange( + 0, (n_real + 1) * step, step, dtype=torch.int32, device=self.b_req_idx.device + ) + req_first = self.b_req_idx.view(n_real, step)[:, 0] + base = (req_first * step).view(n_real, 1) + self.b_ssm_index_rows = base + torch.arange(step, device=base.device, dtype=base.dtype).view(1, step) + assert self.b_ssm_index_rows.shape == (n_real, step) + # The spec conv kernel is per-SEQUENCE (one program per real request), + # indexed by conv_state_indices[idx_seq] with idx_seq in [0, n_real), + # aligned 1:1 with b_gdn_verify_cu_seqlens / b_num_accepted_tokens. The + # default b_conv_buffer_idx = b_req_idx has the expanded length n_real*step, + # which launches n_real*step conv programs and reads num_accepted/ + # query_start_loc out of bounds for idx_seq >= n_real, corrupting the + # committed conv slot. Narrow it to one widened conv slot per request. + self.b_conv_buffer_idx = req_first + return diff --git a/lightllm/common/basemodel/triton_kernel/linear_att_copy.py b/lightllm/common/basemodel/triton_kernel/linear_att_copy.py index d9f631cbd0..fd4c16043d 100644 --- a/lightllm/common/basemodel/triton_kernel/linear_att_copy.py +++ b/lightllm/common/basemodel/triton_kernel/linear_att_copy.py @@ -5,12 +5,13 @@ @triton.jit def _copy_linear_att_state_to_kv_buffer( - gpu_conv_ptr, # [linear_layer_num, size_num, xdim] + gpu_conv_ptr, # [linear_layer_num, size_num, conv_dim * gpu_widened_width] (uint8 tail) gpu_ssm_ptr, # [linear_layer_num, size_num, xxdim] - cpu_kv_conv_ptr, # [size, linear_layer_num, xdim] + cpu_kv_conv_ptr, # [size, linear_layer_num, conv_dim * width_narrow] (uint8 tail) cpu_kv_ssm_ptr, # [size, linear_layer_num, xxdim] b_req_idx, # [batch_size,] big_page_buffer_ids, # [batch_size,] + num_accepted_tokens_ptr, # [batch_size,] gpu_conv_stride_l, gpu_conv_stride_s, gpu_conv_stride_d, @@ -24,7 +25,9 @@ def _copy_linear_att_state_to_kv_buffer( cpu_kv_ssm_stride_l, cpu_kv_ssm_stride_d, mtp_step, - gpu_conv_tail_dim, + conv_dim, # number of conv rows (the d dimension) + gpu_conv_row_bytes, # widened per-row byte length: gpu_widened_width * itemsize + conv_narrow_row_bytes, # narrow per-row byte length: width_narrow * itemsize gpu_ssm_tail_dim, BLOCK: tl.constexpr, ): @@ -40,28 +43,26 @@ def _copy_linear_att_state_to_kv_buffer( return cur_req_idx = tl.load(b_req_idx + cur_batch).to(tl.int64) - cur_state_req_idx = (cur_req_idx * (mtp_step + 1)).to(tl.int64) + accept_len = tl.load(num_accepted_tokens_ptr + cur_batch).to(tl.int64) + canonical_off = accept_len - 1 - for i in range(tl.cdiv(gpu_conv_tail_dim, BLOCK)): - gpu_start_off = i * BLOCK + tl.arange(0, BLOCK) - mask = gpu_start_off < gpu_conv_tail_dim - conv_data = tl.load( - gpu_conv_ptr + cur_layer * gpu_conv_stride_l + cur_state_req_idx * gpu_conv_stride_s + gpu_start_off, - mask=mask, - ) - dest_conv_ptr = ( - cpu_kv_conv_ptr - + big_page_buffer_idx * cpu_kv_conv_stride_s - + cur_layer * cpu_kv_conv_stride_l - + gpu_start_off - ) - tl.store(dest_conv_ptr, conv_data, mask=mask) + conv_src_slot = cur_req_idx + conv_off_bytes = canonical_off * gpu_conv_stride_d + gpu_conv_base = gpu_conv_ptr + cur_layer * gpu_conv_stride_l + conv_src_slot * gpu_conv_stride_s + conv_off_bytes + cpu_conv_base = cpu_kv_conv_ptr + big_page_buffer_idx * cpu_kv_conv_stride_s + cur_layer * cpu_kv_conv_stride_l + for d in range(conv_dim): + for i in range(tl.cdiv(conv_narrow_row_bytes, BLOCK)): + off = i * BLOCK + tl.arange(0, BLOCK) + mask = off < conv_narrow_row_bytes + conv_data = tl.load(gpu_conv_base + d * gpu_conv_row_bytes + off, mask=mask) + tl.store(cpu_conv_base + d * cpu_kv_conv_stride_d + off, conv_data, mask=mask) + ssm_src_slot = (cur_req_idx * (mtp_step + 1) + canonical_off).to(tl.int64) for i in range(tl.cdiv(gpu_ssm_tail_dim, BLOCK)): gpu_start_off = i * BLOCK + tl.arange(0, BLOCK) mask = gpu_start_off < gpu_ssm_tail_dim ssm_data = tl.load( - gpu_ssm_ptr + cur_layer * gpu_ssm_stride_l + cur_state_req_idx * gpu_ssm_stride_s + gpu_start_off, + gpu_ssm_ptr + cur_layer * gpu_ssm_stride_l + ssm_src_slot * gpu_ssm_stride_s + gpu_start_off, mask=mask, ) dest_ssm_ptr = ( @@ -75,32 +76,51 @@ def _copy_linear_att_state_to_kv_buffer( def copy_linear_att_state_to_kv_buffer( b_req_idx: torch.Tensor, big_page_buffer_ids: torch.Tensor, - gpu_conv_state: torch.Tensor, # [linear_layer_num, s, ...] - gpu_ssm_state: torch.Tensor, # [linear_layer_num, s, ...] - cpu_kv_conv_state: torch.Tensor, # [s, linear_layer_num, ...] - cpu_kv_ssm_state: torch.Tensor, # [s, linear_layer_num, ...] + gpu_conv_state: torch.Tensor, # [linear_layer_num, s_widened, conv_dim, gpu_widened_width] + gpu_ssm_state: torch.Tensor, # [linear_layer_num, s_block, ...] + cpu_kv_conv_state: torch.Tensor, # [size, linear_layer_num, conv_dim, width_narrow] + cpu_kv_ssm_state: torch.Tensor, # [size, linear_layer_num, ...] mtp_step: int, + b_num_accepted_tokens: torch.Tensor, # [batch_size,] per-req post-accept count (>=1) ): assert len(b_req_idx) == big_page_buffer_ids.shape[0] + assert len(b_req_idx) == b_num_accepted_tokens.shape[0] BLOCK = 4096 - gpu_conv_state = gpu_conv_state.view(gpu_conv_state.shape[0], gpu_conv_state.shape[1], -1).view(dtype=torch.uint8) - gpu_ssm_state = gpu_ssm_state.view(gpu_ssm_state.shape[0], gpu_ssm_state.shape[1], -1).view(dtype=torch.uint8) - cpu_kv_conv_state = cpu_kv_conv_state.view(cpu_kv_conv_state.shape[0], cpu_kv_conv_state.shape[1], -1).view( - dtype=torch.uint8 + + assert gpu_conv_state.dim() >= 4, "gpu_conv_state must be [layer, s, conv_dim, widened_width]" + assert cpu_kv_conv_state.dim() >= 4, "cpu_kv_conv_state must be [size, layer, conv_dim, width_narrow]" + # #6: the byte snapshot hardcodes gpu_conv_stride_d=conv_itemsize, which is only valid when the + # widened-width axis is element-contiguous (stride 1). Fail fast instead of snapshotting wrong bytes. + assert gpu_conv_state.stride(3) == 1, ( + "gpu_conv_state widened-width axis must be element-contiguous (stride 1); " + "gpu_conv_stride_d=conv_itemsize assumes it" + ) + # #18: canonical_off = accept_len - 1 indexes into the widened slot; bound it to [0, mtp_step] + # (accept_len in [1, mtp_step+1]) so a stale/oversized accept-count can't slice past the slot. + assert int(b_num_accepted_tokens.min()) >= 1 and int(b_num_accepted_tokens.max()) <= mtp_step + 1, ( + f"b_num_accepted_tokens out of range [1, {mtp_step + 1}]: " + f"min={int(b_num_accepted_tokens.min())} max={int(b_num_accepted_tokens.max())}" ) + conv_itemsize = gpu_conv_state.element_size() + gpu_conv_state = gpu_conv_state.view( + gpu_conv_state.shape[0], gpu_conv_state.shape[1], gpu_conv_state.shape[2], -1 + ).view(dtype=torch.uint8) + cpu_kv_conv_state = cpu_kv_conv_state.view( + cpu_kv_conv_state.shape[0], cpu_kv_conv_state.shape[1], cpu_kv_conv_state.shape[2], -1 + ).view(dtype=torch.uint8) + + gpu_ssm_state = gpu_ssm_state.view(gpu_ssm_state.shape[0], gpu_ssm_state.shape[1], -1).view(dtype=torch.uint8) cpu_kv_ssm_state = cpu_kv_ssm_state.view(cpu_kv_ssm_state.shape[0], cpu_kv_ssm_state.shape[1], -1).view( dtype=torch.uint8 ) - assert gpu_conv_state.shape[-1] == cpu_kv_conv_state.shape[-1] + + assert gpu_conv_state.shape[2] == cpu_kv_conv_state.shape[2], "conv_dim mismatch between gpu and cpu conv buffers" assert gpu_ssm_state.shape[-1] == cpu_kv_ssm_state.shape[-1] - assert ( - gpu_conv_state.stride(-1) - == gpu_ssm_state.stride(-1) - == cpu_kv_conv_state.stride(-1) - == cpu_kv_ssm_state.stride(-1) - ) - gpu_conv_tail_dim = gpu_conv_state.shape[-1] + conv_dim = gpu_conv_state.shape[2] + gpu_conv_row_bytes = gpu_conv_state.shape[-1] # widened per-row byte length + conv_narrow_row_bytes = cpu_kv_conv_state.shape[-1] # narrow per-row byte length + assert conv_narrow_row_bytes <= gpu_conv_row_bytes gpu_ssm_tail_dim = gpu_ssm_state.shape[-1] layer_num = gpu_conv_state.shape[0] @@ -114,9 +134,10 @@ def copy_linear_att_state_to_kv_buffer( cpu_kv_ssm_ptr=cpu_kv_ssm_state, b_req_idx=b_req_idx, big_page_buffer_ids=big_page_buffer_ids, + num_accepted_tokens_ptr=b_num_accepted_tokens, gpu_conv_stride_l=gpu_conv_state.stride(0), gpu_conv_stride_s=gpu_conv_state.stride(1), - gpu_conv_stride_d=gpu_conv_state.stride(2), + gpu_conv_stride_d=conv_itemsize, gpu_ssm_stride_l=gpu_ssm_state.stride(0), gpu_ssm_stride_s=gpu_ssm_state.stride(1), gpu_ssm_stride_d=gpu_ssm_state.stride(2), @@ -127,7 +148,9 @@ def copy_linear_att_state_to_kv_buffer( cpu_kv_ssm_stride_l=cpu_kv_ssm_state.stride(1), cpu_kv_ssm_stride_d=cpu_kv_ssm_state.stride(2), mtp_step=mtp_step, - gpu_conv_tail_dim=gpu_conv_tail_dim, + conv_dim=conv_dim, + gpu_conv_row_bytes=gpu_conv_row_bytes, + conv_narrow_row_bytes=conv_narrow_row_bytes, gpu_ssm_tail_dim=gpu_ssm_tail_dim, BLOCK=BLOCK, ) diff --git a/lightllm/common/basemodel/triton_kernel/linear_att_cpu_cache_copy.py b/lightllm/common/basemodel/triton_kernel/linear_att_cpu_cache_copy.py index 37b27cadb2..1251dddc33 100644 --- a/lightllm/common/basemodel/triton_kernel/linear_att_cpu_cache_copy.py +++ b/lightllm/common/basemodel/triton_kernel/linear_att_cpu_cache_copy.py @@ -193,11 +193,7 @@ def copy_kv_buffer_to_cpu_cache( cpu_kv_ssm_tail_dim = cpu_kv_ssm_state.shape[-1] full_att_layer_num = gpu_kv_full_att_state.shape[-2] - assert ( - full_att_layer_num - == (linear_config.all_layer_num // linear_config.full_attention_interval) - == (linear_config.all_layer_num - linear_config.linear_layer_num) - ) + assert full_att_layer_num == linear_config.get_persisted_full_att_layer_num() assert gpu_full_att_tail_dim == cpu_cache_full_att.shape[-1] assert cpu_cache_conv.shape[-1] == cpu_kv_conv_state.shape[-1] assert cpu_cache_ssm.shape[-1] == cpu_kv_ssm_state.shape[-1] @@ -428,6 +424,7 @@ def copy_cpu_cache_to_kv_buffer( cpu_kv_ssm_tail_dim = cpu_kv_ssm_state.shape[-1] full_att_layer_num = gpu_full_att_kv_state.shape[-2] + assert full_att_layer_num == linear_config.get_persisted_full_att_layer_num() assert gpu_full_att_tail_dim == cpu_cache_full_att.shape[-1] assert cpu_cache_conv.shape[-1] == cpu_kv_conv_state.shape[-1] assert cpu_cache_ssm.shape[-1] == cpu_kv_ssm_state.shape[-1] diff --git a/lightllm/common/basemodel/triton_kernel/mtp_utils.py b/lightllm/common/basemodel/triton_kernel/mtp_utils.py index 2d70a68c05..a020605c26 100644 --- a/lightllm/common/basemodel/triton_kernel/mtp_utils.py +++ b/lightllm/common/basemodel/triton_kernel/mtp_utils.py @@ -148,38 +148,6 @@ def mtp_scatter_next_token_ids( ) -@triton.jit -def _fwd_kernel_gen_b_req_mtp_start_loc( - b_mtp_index, - b_req_mtp_start_loc, - num_reqs: tl.constexpr, - batch_size: tl.constexpr, - BLOCK_SIZE: tl.constexpr, -): - offset = tl.arange(0, BLOCK_SIZE) - cur_mtp_index = tl.load(b_mtp_index + offset, mask=offset < batch_size, other=-1) - non_zero_mask = tl.where(cur_mtp_index == 0, 1, 0) # 1 0 1 0 0 - output_offset = tl.cumsum(non_zero_mask) - 1 - tl.store(b_req_mtp_start_loc + output_offset, offset, mask=non_zero_mask == 1) - return - - -def gen_b_req_mtp_start_loc(b_mtp_index: torch.Tensor, num_reqs: int): - b_req_mtp_start_loc = torch.empty((num_reqs,), dtype=torch.int32, device=b_mtp_index.device) - BLOCK_SIZE = triton.next_power_of_2(b_mtp_index.shape[0]) - batch_size = b_mtp_index.shape[0] - grid = (1,) - _fwd_kernel_gen_b_req_mtp_start_loc[grid]( - b_mtp_index=b_mtp_index, - b_req_mtp_start_loc=b_req_mtp_start_loc, - num_reqs=num_reqs, - batch_size=batch_size, - BLOCK_SIZE=BLOCK_SIZE, - num_warps=8, - ) - return b_req_mtp_start_loc - - def test_mtp_verify(): req_to_next_token_ids = torch.tensor( [[1, 2, -2, -1, -1], [1, 2, 0, -1, -1], [1, 3, 4, 4, 5]], dtype=torch.int32, device="cuda" @@ -201,13 +169,5 @@ def test_mtp_verify(): print(accepted_index) -def test_gen_b_req_mtp_start_loc(): - b_mtp_index = torch.tensor([0, 1, 0, 1, 2], dtype=torch.int32, device="cuda") - gt_output = torch.where(b_mtp_index == 0)[0] - b_req_mtp_start_loc = gen_b_req_mtp_start_loc(b_mtp_index, 2) - print(b_req_mtp_start_loc, gt_output) - - if __name__ == "__main__": test_mtp_verify() - # test_gen_b_req_mtp_start_loc() diff --git a/lightllm/common/kv_cache_mem_manager/operator/linear_att.py b/lightllm/common/kv_cache_mem_manager/operator/linear_att.py index 109e813220..e3ae9493c7 100644 --- a/lightllm/common/kv_cache_mem_manager/operator/linear_att.py +++ b/lightllm/common/kv_cache_mem_manager/operator/linear_att.py @@ -24,6 +24,16 @@ def __init__(self, mem_manager): super().__init__(mem_manager) self.linear_config = LinearAttCacheConfig.load_from_args() + @staticmethod + def _get_persisted_full_att_layer_num(mem_manager) -> int: + persisted_full_att = getattr(mem_manager, "persisted_full_att_layer_num", None) + if persisted_full_att is None: + main_full_att = getattr(mem_manager, "main_full_att_layer_num", mem_manager.kv_buffer.shape[0]) + draft_full_att = getattr(mem_manager, "draft_full_att_layers", 0) + persisted_full_att = main_full_att + draft_full_att + assert 0 < persisted_full_att <= mem_manager.kv_buffer.shape[0] + return int(persisted_full_att) + def load_cpu_cache_to_gpu( self, mem_indexes: torch.Tensor, @@ -76,11 +86,14 @@ def load_cpu_cache_to_gpu( copy_cpu_cache_to_kv_buffer, ) + # Restore the persisted full-attn slice: main slots followed by MTP draft slots. + persisted_full_att = self._get_persisted_full_att_layer_num(mem_manager) + copy_cpu_cache_to_kv_buffer( mem_indexes=mem_indexes, big_page_buffer_ids=big_page_buffer_ids_gpu, page_indexes=page_indexes, - gpu_full_att_kv_state=mem_manager.kv_buffer, + gpu_full_att_kv_state=mem_manager.kv_buffer[:persisted_full_att], cpu_kv_conv_state=mem_manager.linear_att_big_page_buffers.conv_state_cache.buffer, cpu_kv_ssm_state=mem_manager.linear_att_big_page_buffers.ssm_state_cache.buffer, cpu_cache_tensor=cpu_cache_client.cpu_kv_cache_tensor, @@ -169,12 +182,15 @@ def offload_gpu_kv_to_cpu_cache( copy_kv_buffer_to_cpu_cache, ) + # Persist the full-attn slice used for prefix reuse: main slots followed by MTP draft slots. + persisted_full_att = self._get_persisted_full_att_layer_num(mem_manager) + copy_kv_buffer_to_cpu_cache( mem_indexes=mem_indexes, page_indexes=page_indexes, page_readies=page_readies, big_page_buffer_ids=big_page_buffer_ids_gpu, - gpu_kv_full_att_state=mem_manager.kv_buffer, + gpu_kv_full_att_state=mem_manager.kv_buffer[:persisted_full_att], cpu_kv_conv_state=mem_manager.linear_att_big_page_buffers.conv_state_cache.buffer, cpu_kv_ssm_state=mem_manager.linear_att_big_page_buffers.ssm_state_cache.buffer, cpu_cache_tensor=cpu_cache_client.cpu_kv_cache_tensor, diff --git a/lightllm/common/linear_att_cache_manager/config_objs.py b/lightllm/common/linear_att_cache_manager/config_objs.py index 46ab9d2107..fdd8ed15c6 100644 --- a/lightllm/common/linear_att_cache_manager/config_objs.py +++ b/lightllm/common/linear_att_cache_manager/config_objs.py @@ -1,13 +1,18 @@ import torch import dataclasses import triton -from lightllm.utils.envs_utils import get_env_start_args +from lightllm.utils.envs_utils import get_env_start_args, _mtp_added_layer_num from lightllm.utils.log_utils import init_logger from lightllm.utils.torch_dtype_utils import get_torch_dtype logger = init_logger(__name__) +def get_mtp_draft_full_att_layer_num(args) -> int: + # Delegates to the single source of truth in envs_utils (#9). + return _mtp_added_layer_num(getattr(args, "mtp_mode", None), getattr(args, "mtp_step", 0)) + + @dataclasses.dataclass class LinearAttCacheConfig: tp_world_size: int @@ -28,13 +33,30 @@ class LinearAttCacheConfig: ssm_state_dtype: torch.dtype full_attention_interval: int all_layer_num: int # 包括 linear att 和 full att 的层加起来的层数 + draft_full_att_layer_num: int = 0 def get_conv_dim(self): return self.head_linear_k_dim * self.num_linear_k_heads * 2 + self.head_linear_v_dim * self.num_linear_v_heads - def get_conv_state_shape(self): + def get_main_full_att_layer_num(self): + main_full_att_layer_num = self.all_layer_num - self.linear_layer_num + assert main_full_att_layer_num == self.all_layer_num // self.full_attention_interval + return main_full_att_layer_num + + def get_persisted_full_att_layer_num(self): + return self.get_main_full_att_layer_num() + self.draft_full_att_layer_num + + def get_persisted_conv_state_shape(self): + # NARROW shape used for the CPU/disk persisted page and ALL byte math. + # Persisted state is always the committed (narrow) sliding window. return (self.get_conv_dim(), self.conv_kernel_size - 1) + def get_gpu_conv_state_shape(self, mtp_step: int): + # WIDENED working shape for the GPU buffer: holds the tentatively + # rolled-in S speculative tokens before acceptance. width-1 + S, where + # S = mtp_step (a verify step has seqlen=S+1 -> width-1+(seqlen-1)). + return (self.get_conv_dim(), (self.conv_kernel_size - 1) + mtp_step) + def get_ssm_state_shape(self): return (self.num_linear_v_heads, self.head_linear_k_dim, self.head_linear_v_dim) @@ -57,7 +79,7 @@ def get_cpu_cache_full_att_bytes(self): ) assert big_page_token_num == get_env_start_args().cpu_cache_token_page_size full_att_bytes = 2 * self.full_att_all_num_kv_heads * self.full_att_head_dim * self.full_att_dtype.itemsize - a = full_att_bytes * (self.all_layer_num - self.linear_layer_num) * big_page_token_num + a = full_att_bytes * self.get_persisted_full_att_layer_num() * big_page_token_num return a def get_cpu_cache_conv_bytes(self): @@ -102,4 +124,5 @@ def load_from_args() -> "LinearAttCacheConfig": ssm_state_dtype=get_torch_dtype(args.linear_att_ssm_data_type), full_attention_interval=llm_config["full_attention_interval"], all_layer_num=n_layer, + draft_full_att_layer_num=get_mtp_draft_full_att_layer_num(args), ) diff --git a/lightllm/common/linear_att_cache_manager/linear_att_buffer_manager.py b/lightllm/common/linear_att_cache_manager/linear_att_buffer_manager.py index 30dc4d937c..2ab4313e37 100644 --- a/lightllm/common/linear_att_cache_manager/linear_att_buffer_manager.py +++ b/lightllm/common/linear_att_cache_manager/linear_att_buffer_manager.py @@ -24,7 +24,7 @@ def __init__( self.conv_state_cache = LayerCache( size=self.size, dtype=self.linear_config.conv_state_dtype, - shape=self.linear_config.get_conv_state_shape(), + shape=self.linear_config.get_persisted_conv_state_shape(), layer_num=self.linear_config.linear_layer_num, device="cpu", size_first=True, diff --git a/lightllm/common/req_manager.py b/lightllm/common/req_manager.py index 01e9c4ad35..6673243c9f 100644 --- a/lightllm/common/req_manager.py +++ b/lightllm/common/req_manager.py @@ -19,6 +19,18 @@ logger = init_logger(__name__) +# Width of req_to_next_token_ids: holds the seed token + up to (WIDTH - 1) MTP draft tokens. +REQ_NEXT_TOKEN_IDS_WIDTH = 8 + + +def assert_mtp_step_within_next_token_ids_width(mtp_step: int) -> None: + assert mtp_step <= REQ_NEXT_TOKEN_IDS_WIDTH - 1, ( + f"mtp_step={mtp_step} exceeds {REQ_NEXT_TOKEN_IDS_WIDTH - 1}; " + f"req_to_next_token_ids width is {REQ_NEXT_TOKEN_IDS_WIDTH} " + "(widening it is an explicit follow-up, spec §9)" + ) + + class _ReqNode: def __init__(self, index): self.index = index @@ -117,7 +129,7 @@ def __init__(self, max_request_num): self.req_to_frequency_penalty = torch.zeros(max_request_num + 1, dtype=torch.float32, device="cuda") self.req_to_repetition_penalty = torch.zeros(max_request_num + 1, dtype=torch.float32, device="cuda") self.req_to_next_token_ids = torch.zeros( - (max_request_num + 1, 8), + (max_request_num + 1, REQ_NEXT_TOKEN_IDS_WIDTH), dtype=torch.int64, device="cuda", ) @@ -236,15 +248,13 @@ def __init__(self, max_request_num, max_sequence_length, mem_manager, linear_con self.big_page_token_num = ( get_env_start_args().linear_att_page_block_num * get_env_start_args().linear_att_hash_page_size ) - assert ( - self.mtp_step == 0 - ), "currently only support mtp_step 0 for simplicity, more mtp_step support will be added in the future" + assert_mtp_step_within_next_token_ids_width(self.mtp_step) self.linear_config = linear_config self.req_to_conv_state = LayerCache( - size=(max_request_num + 1) * (self.mtp_step + 1), + size=(max_request_num + 1), dtype=self.linear_config.conv_state_dtype, - shape=self.linear_config.get_conv_state_shape(), + shape=self.linear_config.get_gpu_conv_state_shape(mtp_step=self.mtp_step), layer_num=self.linear_config.linear_layer_num, device="cuda", ) @@ -258,11 +268,13 @@ def __init__(self, max_request_num, max_sequence_length, mem_manager, linear_con return def init_linear_att_state(self, req: "InferReq"): - index = req.req_idx * (self.mtp_step + 1) - conv_state = self.req_to_conv_state.buffer[:, index, ...] - ssm_state = self.req_to_ssm_state.buffer[:, index, ...] - conv_state.fill_(0) - ssm_state.fill_(0) + conv_index = req.req_idx + ssm_start = req.req_idx * (self.mtp_step + 1) + self.req_to_conv_state.buffer[:, conv_index, ...].fill_(0) + # #17: zero the FULL (mtp_step + 1)-row SSM block, not just canonical row +0, so a future + # first-step verify reading offset>0 after fresh init never hits a never-written row (NaN). + self.req_to_ssm_state.buffer[:, ssm_start : ssm_start + (self.mtp_step + 1), ...].fill_(0) + req.mtp_accept_len = 1 return def get_mamba_cache(self, layer_idx_in_all: int): @@ -275,16 +287,17 @@ def get_mamba_cache(self, layer_idx_in_all: int): return conv_states, ssm_states def copy_big_page_buffer_to_linear_att_state(self, big_page_buffer_idx: int, req: "InferReq"): - from .linear_att_cache_manager import LinearAttCacheManager big_page_buffers: LinearAttCacheManager = self.mem_manager.linear_att_big_page_buffers conv_state, ssm_state = big_page_buffers.get_state_cache(buffer_idx=big_page_buffer_idx) - dest_req_idx = req.req_idx * (self.mtp_step + 1) - - self.req_to_conv_state.buffer[:, dest_req_idx, ...] = conv_state - self.req_to_ssm_state.buffer[:, dest_req_idx, ...] = ssm_state + conv_dest = req.req_idx + ssm_dest = req.req_idx * (self.mtp_step + 1) + narrow_w = conv_state.shape[-1] # persisted (narrow) width + self.req_to_conv_state.buffer[:, conv_dest, ..., :narrow_w] = conv_state + self.req_to_ssm_state.buffer[:, ssm_dest, ...] = ssm_state + req.mtp_accept_len = 1 return def copy_small_page_buffer_to_linear_att_state( @@ -293,9 +306,12 @@ def copy_small_page_buffer_to_linear_att_state( conv_state, ssm_state = linear_att_small_page_buffers.get_state_cache( buffer_idx=req.shared_kv_node.small_page_buffer_idx ) - dest_req_idx = req.req_idx * (self.mtp_step + 1) + conv_dest = req.req_idx + ssm_dest = req.req_idx * (self.mtp_step + 1) + narrow_w = conv_state.shape[-1] # TODO 下面这个从 cpu cache 拷贝数据的 gpu的操作,是否是阻塞的操作。 # 同时,非连续对象的拷贝,可能存在效率问题。 - self.req_to_conv_state.buffer[:, dest_req_idx, ...] = conv_state - self.req_to_ssm_state.buffer[:, dest_req_idx, ...] = ssm_state + self.req_to_conv_state.buffer[:, conv_dest, ..., :narrow_w] = conv_state + self.req_to_ssm_state.buffer[:, ssm_dest, ...] = ssm_state + req.mtp_accept_len = 1 return diff --git a/lightllm/models/base_mtp_model.py b/lightllm/models/base_mtp_model.py new file mode 100644 index 0000000000..796a34d9ad --- /dev/null +++ b/lightllm/models/base_mtp_model.py @@ -0,0 +1,36 @@ +from typing import List + +from lightllm.common.basemodel.basemodel import TpPartBaseModel + + +class BaseMTPModel: + """Shared wiring for MTP draft models: they reuse the main model's req/mem managers and rope + caches, and pop the main_model / previous-draft-models kwargs before the base __init__ (#25). + Mixed in BEFORE the concrete base model so these overrides win via MRO. + + Also carries the is_mtp_draft_model marker consumed by detection sites (#23).""" + + is_mtp_draft_model = True + + def __init__(self, kvargs: dict): + self._pre_init(kvargs) + super().__init__(kvargs) + return + + def _pre_init(self, kvargs: dict): + self.main_model: TpPartBaseModel = kvargs.pop("main_model") + self.mtp_previous_draft_models: List[TpPartBaseModel] = kvargs.pop("mtp_previous_draft_models") + return + + def _init_custom(self): + self._cos_cached = self.main_model._cos_cached + self._sin_cached = self.main_model._sin_cached + return + + def _init_req_manager(self): + self.req_manager = self.main_model.req_manager + return + + def _init_mem_manager(self): + self.mem_manager = self.main_model.mem_manager + return diff --git a/lightllm/models/deepseek_mtp/model.py b/lightllm/models/deepseek_mtp/model.py index d9ffdb0e31..369ff1766f 100644 --- a/lightllm/models/deepseek_mtp/model.py +++ b/lightllm/models/deepseek_mtp/model.py @@ -1,38 +1,14 @@ -from typing import List +from lightllm.models.base_mtp_model import BaseMTPModel from lightllm.models.deepseek2.model import Deepseek2TpPartModel from lightllm.models.deepseek_mtp.layer_infer.pre_layer_infer import Deepseek3MTPPreLayerInfer from lightllm.models.deepseek_mtp.layer_weights.pre_and_post_layer_weight import Deepseek3MTPPreAndPostLayerWeight -from lightllm.common.basemodel import TpPartBaseModel -class Deepseek3MTPModel(Deepseek2TpPartModel): +class Deepseek3MTPModel(BaseMTPModel, Deepseek2TpPartModel): pre_and_post_weight_class = Deepseek3MTPPreAndPostLayerWeight pre_layer_infer_class = Deepseek3MTPPreLayerInfer - def __init__(self, kvargs: dict): - self._pre_init(kvargs) - super().__init__(kvargs) - return - - def _pre_init(self, kvargs: dict): - self.main_model: TpPartBaseModel = kvargs.pop("main_model") - self.mtp_previous_draft_models: List[TpPartBaseModel] = kvargs.pop("mtp_previous_draft_models") - return - - def _init_custom(self): - self._cos_cached = self.main_model._cos_cached - self._sin_cached = self.main_model._sin_cached - return - - def _init_req_manager(self): - self.req_manager = self.main_model.req_manager - return - - def _init_mem_manager(self): - self.mem_manager = self.main_model.mem_manager - return - def _init_weights(self, start_layer_index=None): assert start_layer_index is None self.pre_post_weight = self.pre_and_post_weight_class( diff --git a/lightllm/models/glm4_moe_lite_mtp/model.py b/lightllm/models/glm4_moe_lite_mtp/model.py index 549bf7ce41..039491dc1f 100644 --- a/lightllm/models/glm4_moe_lite_mtp/model.py +++ b/lightllm/models/glm4_moe_lite_mtp/model.py @@ -1,36 +1,17 @@ -from typing import List +from lightllm.models.base_mtp_model import BaseMTPModel from lightllm.models.deepseek_mtp.layer_infer.pre_layer_infer import Deepseek3MTPPreLayerInfer from lightllm.models.glm4_moe_lite.model import Glm4MoeLiteTpPartModel from lightllm.models.glm4_moe_lite_mtp.layer_weights.pre_and_post_layer_weight import ( Glm4MoeLiteMTPPreAndPostLayerWeight, ) -from lightllm.common.basemodel import TpPartBaseModel from lightllm.common.basemodel.basemodel import load_hf_weights -class Glm4MoeLiteMTPModel(Glm4MoeLiteTpPartModel): +class Glm4MoeLiteMTPModel(BaseMTPModel, Glm4MoeLiteTpPartModel): pre_and_post_weight_class = Glm4MoeLiteMTPPreAndPostLayerWeight pre_layer_infer_class = Deepseek3MTPPreLayerInfer - def __init__(self, kvargs: dict): - self._pre_init(kvargs) - super().__init__(kvargs) - - def _pre_init(self, kvargs: dict): - self.main_model: TpPartBaseModel = kvargs.pop("main_model") - self.mtp_previous_draft_models: List[TpPartBaseModel] = kvargs.pop("mtp_previous_draft_models") - - def _init_custom(self): - self._cos_cached = self.main_model._cos_cached - self._sin_cached = self.main_model._sin_cached - - def _init_req_manager(self): - self.req_manager = self.main_model.req_manager - - def _init_mem_manager(self): - self.mem_manager = self.main_model.mem_manager - def _init_weights(self, start_layer_index=None): assert start_layer_index is None diff --git a/lightllm/models/mistral_mtp/model.py b/lightllm/models/mistral_mtp/model.py index 7c64625ca8..e1edd9d027 100644 --- a/lightllm/models/mistral_mtp/model.py +++ b/lightllm/models/mistral_mtp/model.py @@ -1,14 +1,13 @@ -from typing import List +from lightllm.models.base_mtp_model import BaseMTPModel from lightllm.models.mistral.model import MistralTpPartModel from lightllm.models.mistral_mtp.layer_weights.pre_and_post_layer_weight import MistralMTPPreAndPostLayerWeight from lightllm.models.mistral_mtp.layer_infer.pre_layer_infer import MistralMTPPreLayerInfer from lightllm.models.mistral_mtp.layer_infer.post_layer_infer import MistralMTPPostLayerInfer from lightllm.models.mistral_mtp.layer_infer.transformer_layer_infer import MistralMTPTransformerLayerInfer from lightllm.models.mistral_mtp.layer_weights.transformer_layer_weight import MistralMTPTransformerLayerWeight -from lightllm.common.basemodel import TpPartBaseModel -class MistralMTPModel(MistralTpPartModel): +class MistralMTPModel(BaseMTPModel, MistralTpPartModel): pre_and_post_weight_class = MistralMTPPreAndPostLayerWeight pre_layer_infer_class = MistralMTPPreLayerInfer @@ -18,34 +17,11 @@ class MistralMTPModel(MistralTpPartModel): post_layer_infer_class = MistralMTPPostLayerInfer - def __init__(self, kvargs: dict): - self._pre_init(kvargs) - super().__init__(kvargs) - return - - def _pre_init(self, kvargs: dict): - self.main_model: TpPartBaseModel = kvargs.pop("main_model") - self.mtp_previous_draft_models: List[TpPartBaseModel] = kvargs.pop("mtp_previous_draft_models") - return - def _init_some_value(self): super()._init_some_value() self.layers_num = 1 return - def _init_custom(self): - self._cos_cached = self.main_model._cos_cached - self._sin_cached = self.main_model._sin_cached - return - - def _init_req_manager(self): - self.req_manager = self.main_model.req_manager - return - - def _init_mem_manager(self): - self.mem_manager = self.main_model.mem_manager - return - def _init_weights(self, start_layer_index=None): assert start_layer_index is None self.config["n_layer"] = 1 diff --git a/lightllm/models/qwen3_5/infer_struct.py b/lightllm/models/qwen3_5/infer_struct.py index d23475c1cf..2687a4aca7 100644 --- a/lightllm/models/qwen3_5/infer_struct.py +++ b/lightllm/models/qwen3_5/infer_struct.py @@ -1,8 +1,4 @@ -import torch -from typing import List - from lightllm.models.qwen2_vl.infer_struct import Qwen2VLInferStateInfo -from lightllm.utils.envs_utils import get_env_start_args class Qwen35InferStateInfo(Qwen2VLInferStateInfo): @@ -12,8 +8,7 @@ def __init__(self): def init_some_extra_state(self, model): super().init_some_extra_state(model) - self.b_att_seq_len = self.b_seq_len - mtp_step = get_env_start_args().mtp_step + from lightllm.common.basemodel.mtp_verify_extra_state import init_mtp_verify_extra_state - self.b_buffer_idx = self.b_req_idx * (mtp_step + 1) + self.b_mtp_index + init_mtp_verify_extra_state(self) return diff --git a/lightllm/models/qwen3_5_moe_mtp/__init__.py b/lightllm/models/qwen3_5_moe_mtp/__init__.py new file mode 100644 index 0000000000..c8885f8869 --- /dev/null +++ b/lightllm/models/qwen3_5_moe_mtp/__init__.py @@ -0,0 +1,3 @@ +from lightllm.models.qwen3_5_moe_mtp.model import Qwen3_5MoeMTPModel + +__all__ = ["Qwen3_5MoeMTPModel"] diff --git a/lightllm/models/qwen3_5_moe_mtp/layer_weights/__init__.py b/lightllm/models/qwen3_5_moe_mtp/layer_weights/__init__.py new file mode 100644 index 0000000000..dcad1087d4 --- /dev/null +++ b/lightllm/models/qwen3_5_moe_mtp/layer_weights/__init__.py @@ -0,0 +1,5 @@ +from lightllm.models.qwen3_5_moe_mtp.layer_weights.transformer_layer_weight import ( + Qwen3_5MoeMTPTransformerLayerWeight, +) + +__all__ = ["Qwen3_5MoeMTPTransformerLayerWeight"] diff --git a/lightllm/models/qwen3_5_moe_mtp/layer_weights/transformer_layer_weight.py b/lightllm/models/qwen3_5_moe_mtp/layer_weights/transformer_layer_weight.py new file mode 100644 index 0000000000..554db359f6 --- /dev/null +++ b/lightllm/models/qwen3_5_moe_mtp/layer_weights/transformer_layer_weight.py @@ -0,0 +1,96 @@ +from lightllm.common.basemodel.layer_weights.meta_weights import ( + COLMMWeight, + FusedMoeWeight, + ROWMMWeight, +) +from lightllm.models.qwen3_5_moe.layer_weights.transformer_layer_weight import ( + Qwen35MOETransformerLayerWeight, +) +from lightllm.models.qwen3_5_mtp.layer_weights.mtp_retarget_mixin import MTPRetargetMixin +from lightllm.utils.envs_utils import get_env_start_args + + +class Qwen3_5MoeMTPTransformerLayerWeight(MTPRetargetMixin, Qwen35MOETransformerLayerWeight): + def _init_weight_names(self): + super()._init_weight_names() + self._retarget_attn_norm_names() + + def _init_moe(self): + moe_intermediate_size = self.network_config_["moe_intermediate_size"] + self.moe_gate = ROWMMWeight( + in_dim=self.network_config_["hidden_size"], + out_dims=[self.n_routed_experts], + weight_names=f"{self._MTP_PREFIX}{self.layer_num_}.mlp.gate.weight", + data_type=self.data_type_, + quant_method=None, + tp_rank=0, + tp_world_size=1, + ) + self.experts = FusedMoeWeight( + gate_proj_name="gate_proj", + down_proj_name="down_proj", + up_proj_name="up_proj", + e_score_correction_bias_name="", + weight_prefix=f"{self._MTP_PREFIX}{self.layer_num_}.mlp.experts", + n_routed_experts=self.n_routed_experts, + hidden_size=self.network_config_["hidden_size"], + moe_intermediate_size=moe_intermediate_size, + data_type=self.data_type_, + quant_method=self.quant_cfg.get_quant_method(self.layer_num_, "fused_moe"), + layer_num=self.layer_num_, + network_config=self.network_config_, + ) + self._init_gated_ffn() + + def _init_gated_ffn(self): + hidden_size = self.network_config_["hidden_size"] + if "shared_expert_intermediate_size" not in self.network_config_: + return + + prefix = f"{self._MTP_PREFIX}{self.layer_num_}.mlp.shared_expert" + inter_size = self.network_config_["shared_expert_intermediate_size"] + if get_env_start_args().enable_ep_moe: + self.gate_up_proj = ROWMMWeight( + in_dim=hidden_size, + out_dims=[inter_size, inter_size], + weight_names=[f"{prefix}.gate_proj.weight", f"{prefix}.up_proj.weight"], + data_type=self.data_type_, + quant_method=self.get_quant_method("gate_up_proj"), + tp_rank=0, + tp_world_size=1, + ) + self.down_proj = COLMMWeight( + in_dim=inter_size, + out_dims=[hidden_size], + weight_names=f"{prefix}.down_proj.weight", + data_type=self.data_type_, + quant_method=self.get_quant_method("down_proj"), + tp_rank=0, + tp_world_size=1, + ) + else: + self.gate_up_proj = ROWMMWeight( + in_dim=hidden_size, + out_dims=[inter_size, inter_size], + weight_names=[f"{prefix}.gate_proj.weight", f"{prefix}.up_proj.weight"], + data_type=self.data_type_, + quant_method=self.get_quant_method("gate_up_proj"), + ) + self.down_proj = COLMMWeight( + in_dim=inter_size, + out_dims=[hidden_size], + weight_names=f"{prefix}.down_proj.weight", + data_type=self.data_type_, + quant_method=self.get_quant_method("down_proj"), + ) + + self.ffn_gate = ROWMMWeight( + in_dim=hidden_size, + out_dims=[1], + weight_names=f"{self._MTP_PREFIX}{self.layer_num_}.mlp.shared_expert_gate.weight", + data_type=self.data_type_, + bias_names=None, + quant_method=None, + tp_rank=0, + tp_world_size=1, + ) diff --git a/lightllm/models/qwen3_5_moe_mtp/model.py b/lightllm/models/qwen3_5_moe_mtp/model.py new file mode 100644 index 0000000000..022864f6b3 --- /dev/null +++ b/lightllm/models/qwen3_5_moe_mtp/model.py @@ -0,0 +1,8 @@ +from lightllm.models.qwen3_5_mtp.model import Qwen3_5MTPModel +from lightllm.models.qwen3_5_moe_mtp.layer_weights.transformer_layer_weight import ( + Qwen3_5MoeMTPTransformerLayerWeight, +) + + +class Qwen3_5MoeMTPModel(Qwen3_5MTPModel): + transformer_weight_class = Qwen3_5MoeMTPTransformerLayerWeight diff --git a/lightllm/models/qwen3_5_mtp/__init__.py b/lightllm/models/qwen3_5_mtp/__init__.py new file mode 100644 index 0000000000..e69de29bb2 diff --git a/lightllm/models/qwen3_5_mtp/layer_infer/__init__.py b/lightllm/models/qwen3_5_mtp/layer_infer/__init__.py new file mode 100644 index 0000000000..e69de29bb2 diff --git a/lightllm/models/qwen3_5_mtp/layer_infer/pre_layer_infer.py b/lightllm/models/qwen3_5_mtp/layer_infer/pre_layer_infer.py new file mode 100644 index 0000000000..906a0ab62c --- /dev/null +++ b/lightllm/models/qwen3_5_mtp/layer_infer/pre_layer_infer.py @@ -0,0 +1,40 @@ +import torch + +from lightllm.models.qwen3_vl.layer_infer.pre_layer_infer import Qwen3VLMultimodalPreLayerInfer +from lightllm.models.qwen3_5_mtp.layer_weights.pre_and_post_layer_weight import Qwen3_5MTPPreAndPostLayerWeight +from lightllm.models.llama.infer_struct import LlamaInferStateInfo + + +class Qwen3_5MTPPreLayerInfer(Qwen3VLMultimodalPreLayerInfer): + def __init__(self, network_config): + super().__init__(network_config) + self.eps_ = network_config["rms_norm_eps"] + self.hidden_size = network_config["hidden_size"] + return + + def _mtp_fuse( + self, + input_embdings: torch.Tensor, + infer_state: LlamaInferStateInfo, + layer_weight: Qwen3_5MTPPreAndPostLayerWeight, + ) -> torch.Tensor: + tgt_embdings = infer_state.mtp_draft_input_hiddens + assert ( + input_embdings.shape[0] == tgt_embdings.shape[0] + ), f"shape {input_embdings.shape} != shape {tgt_embdings.shape}" + + layer_weight.enorm_weight_(input=input_embdings, eps=self.eps_, out=input_embdings) + layer_weight.hnorm_weight_(input=tgt_embdings, eps=self.eps_, out=tgt_embdings) + cat_embdings = torch.cat((input_embdings, tgt_embdings), dim=-1) + + return layer_weight.eh_proj_weight_.mm(cat_embdings) + + def context_forward( + self, input_ids, infer_state: LlamaInferStateInfo, layer_weight: Qwen3_5MTPPreAndPostLayerWeight + ): + input_embdings = super().context_forward(input_ids, infer_state, layer_weight) + return self._mtp_fuse(input_embdings, infer_state, layer_weight) + + def token_forward(self, input_ids, infer_state: LlamaInferStateInfo, layer_weight: Qwen3_5MTPPreAndPostLayerWeight): + input_embdings = super().token_forward(input_ids, infer_state, layer_weight) + return self._mtp_fuse(input_embdings, infer_state, layer_weight) diff --git a/lightllm/models/qwen3_5_mtp/layer_weights/__init__.py b/lightllm/models/qwen3_5_mtp/layer_weights/__init__.py new file mode 100644 index 0000000000..e69de29bb2 diff --git a/lightllm/models/qwen3_5_mtp/layer_weights/mtp_retarget_mixin.py b/lightllm/models/qwen3_5_mtp/layer_weights/mtp_retarget_mixin.py new file mode 100644 index 0000000000..cf9da94887 --- /dev/null +++ b/lightllm/models/qwen3_5_mtp/layer_weights/mtp_retarget_mixin.py @@ -0,0 +1,61 @@ +from lightllm.common.basemodel.layer_weights.meta_weights import ROWMMWeight, QKVROWNMMWeight + + +class MTPRetargetMixin: + """Shared MTP weight-name retargeting (model.layers.* -> mtp.layers.*) and qkv/o_gate wiring, + used by both the dense and MoE Qwen3.5 MTP layer-weight classes (#11). The dense subclass adds + its dense-MLP retargets on top; the MoE subclass must not (it uses fused experts).""" + + _MAIN_PREFIX = "model.layers." + _MTP_PREFIX = "mtp.layers." + + _ATTN_NORM_NAME_ATTRS = ( + "_q_weight_name", + "_q_norm_name", + "_q_bias_name", + "_k_weight_name", + "_k_norm_name", + "_k_bias_name", + "_v_weight_name", + "_v_bias_name", + "_kv_weight_name", + "_kv_bias_name", + "_o_weight_name", + "_o_bias_name", + "_att_norm_weight_name", + "_att_norm_bias_name", + "_ffn_norm_weight_name", + "_ffn_norm_bias_name", + ) + + def _retarget(self, name): + if name is None: + return None + return name.replace(self._MAIN_PREFIX, self._MTP_PREFIX, 1) + + def _retarget_attn_norm_names(self): + for attr in self._ATTN_NORM_NAME_ATTRS: + setattr(self, attr, self._retarget(getattr(self, attr))) + + def _init_qkv(self): + in_dim = self.n_embed + q_out_dim = self.q_head_num_ * self.head_dim + self.qkv_proj = QKVROWNMMWeight( + in_dim=in_dim, + q_head_num=self.q_head_num_, + kv_head_num=self.k_head_num_, + head_dim=self.head_dim, + weight_names=[self._q_weight_name, self._k_weight_name, self._v_weight_name], + data_type=self.data_type_, + bias_names=[self._q_bias_name, self._k_bias_name, self._v_bias_name], + quant_method=self.get_quant_method("qkv_proj"), + ) + self._o_gate_weight_name = f"{self._MTP_PREFIX}{self.layer_num_}.self_attn.o_gate_proj.weight" + self._o_gate_proj = ROWMMWeight( + in_dim=in_dim, + out_dims=[q_out_dim], + weight_names=[self._o_gate_weight_name], + data_type=self.data_type_, + bias_names=None, + quant_method=self.get_quant_method("o_gate_proj"), + ) diff --git a/lightllm/models/qwen3_5_mtp/layer_weights/pre_and_post_layer_weight.py b/lightllm/models/qwen3_5_mtp/layer_weights/pre_and_post_layer_weight.py new file mode 100644 index 0000000000..25c56a0d7e --- /dev/null +++ b/lightllm/models/qwen3_5_mtp/layer_weights/pre_and_post_layer_weight.py @@ -0,0 +1,45 @@ +from lightllm.common.basemodel import PreAndPostLayerWeight +from lightllm.common.basemodel.layer_weights.meta_weights import ( + EmbeddingWeight, + LMHeadWeight, + NoTpGEMMANormWeight, + ROWMMWeight, +) +from lightllm.common.quantization import Quantcfg + + +class Qwen3_5MTPPreAndPostLayerWeight(PreAndPostLayerWeight): + def __init__(self, data_type, network_config, quant_cfg: Quantcfg): + super().__init__(data_type, network_config) + self.quant_cfg: Quantcfg = quant_cfg + hidden_size = network_config["hidden_size"] + + self.eh_proj_weight_ = ROWMMWeight( + in_dim=hidden_size * 2, + out_dims=[hidden_size], + weight_names="mtp.fc.weight", + data_type=self.data_type_, + quant_method=self.quant_cfg.get_quant_method(0, "eh_proj"), + tp_rank=0, + tp_world_size=1, + ) + self.enorm_weight_ = NoTpGEMMANormWeight( + dim=hidden_size, + weight_name="mtp.pre_fc_norm_embedding.weight", + data_type=self.data_type_, + ) + self.hnorm_weight_ = NoTpGEMMANormWeight( + dim=hidden_size, + weight_name="mtp.pre_fc_norm_hidden.weight", + data_type=self.data_type_, + ) + self.final_norm_weight_ = NoTpGEMMANormWeight( + dim=hidden_size, + weight_name="mtp.norm.weight", + data_type=self.data_type_, + ) + + # Shared with the main Qwen3.5 model, injected by the model class (not loaded here). + self.wte_weight_: EmbeddingWeight = None + self.lm_head_weight_: LMHeadWeight = None + return diff --git a/lightllm/models/qwen3_5_mtp/layer_weights/transformer_layer_weight.py b/lightllm/models/qwen3_5_mtp/layer_weights/transformer_layer_weight.py new file mode 100644 index 0000000000..5aa0724580 --- /dev/null +++ b/lightllm/models/qwen3_5_mtp/layer_weights/transformer_layer_weight.py @@ -0,0 +1,23 @@ +from lightllm.models.qwen3_5.layer_weights.transformer_layer_weight import ( + Qwen35TransformerLayerWeight, +) +from lightllm.models.qwen3_5_mtp.layer_weights.mtp_retarget_mixin import MTPRetargetMixin +from lightllm.utils.log_utils import init_logger + +logger = init_logger(__name__) + + +class Qwen3_5MTPTransformerLayerWeight(MTPRetargetMixin, Qwen35TransformerLayerWeight): + def _init_weight_names(self): + super()._init_weight_names() + # Retarget all main-model layer key names to the mtp.* namespace. + self._retarget_attn_norm_names() + # MLP (dense) projection names retargeted by Qwen35TransformerLayerWeight. + self._gate_weight_name = self._retarget(self._gate_weight_name) + self._gate_bias_name = self._retarget(self._gate_bias_name) + self._up_weight_name = self._retarget(self._up_weight_name) + self._up_bias_name = self._retarget(self._up_bias_name) + self._gate_up_weight_name = self._retarget(self._gate_up_weight_name) + self._gate_up_bias_name = self._retarget(self._gate_up_bias_name) + self._down_weight_name = self._retarget(self._down_weight_name) + self._down_bias_name = self._retarget(self._down_bias_name) diff --git a/lightllm/models/qwen3_5_mtp/model.py b/lightllm/models/qwen3_5_mtp/model.py new file mode 100644 index 0000000000..44ca5b2f1e --- /dev/null +++ b/lightllm/models/qwen3_5_mtp/model.py @@ -0,0 +1,80 @@ +from lightllm.models.base_mtp_model import BaseMTPModel +from lightllm.models.qwen3_5.model import Qwen3_5TpPartModel +from lightllm.models.qwen3_5.layer_infer.transformer_layer_infer import Qwen35TransformerLayerInfer +from lightllm.models.qwen3_5_mtp.layer_weights.pre_and_post_layer_weight import Qwen3_5MTPPreAndPostLayerWeight +from lightllm.models.qwen3_5_mtp.layer_weights.transformer_layer_weight import Qwen3_5MTPTransformerLayerWeight +from lightllm.models.qwen3_5_mtp.layer_infer.pre_layer_infer import Qwen3_5MTPPreLayerInfer +from lightllm.utils.log_utils import init_logger + +logger = init_logger(__name__) + + +class Qwen3_5MTPModel(BaseMTPModel, Qwen3_5TpPartModel): + + pre_and_post_weight_class = Qwen3_5MTPPreAndPostLayerWeight + pre_layer_infer_class = Qwen3_5MTPPreLayerInfer + transformer_weight_class = Qwen3_5MTPTransformerLayerWeight + transformer_layer_infer_class = Qwen35TransformerLayerInfer + + def _init_config(self): + super()._init_config() + self.config["full_attention_interval"] = 1 + self.config["num_hidden_layers"] = 1 + self.config["n_layer"] = 1 + return + + def _init_some_value(self): + super()._init_some_value() + self.layers_num = 1 + return + + def _init_weights(self, start_layer_index=None): + assert start_layer_index is None + self.pre_post_weight = self.pre_and_post_weight_class( + self.data_type, network_config=self.config, quant_cfg=self.quant_cfg + ) + self.trans_layers_weight = [ + self.transformer_weight_class( + i, + self.data_type, + network_config=self.config, + quant_cfg=self.quant_cfg, + ) + for i in range(0, self.config["n_layer"]) + ] + # Shared with the main Qwen3.5 model (mtp_use_dedicated_embeddings: false). + self.pre_post_weight.wte_weight_ = self.main_model.pre_post_weight.wte_weight_ + self.pre_post_weight.lm_head_weight_ = self.main_model.pre_post_weight.lm_head_weight_ + return + + def _init_infer_layer(self, start_layer_index=None): + assert start_layer_index is None + # Build the single draft layer with layer_num == 0 so that, with + # full_attention_interval == 1, it takes the full-attention (mrope) path. + super()._init_infer_layer(start_layer_index=0) + self._assign_draft_kv_slot() + return + + def _assign_draft_kv_slot(self): + mem_manager = self.main_model.mem_manager + main_full_att = getattr(mem_manager, "main_full_att_layer_num", None) + interval = self.main_model.config["full_attention_interval"] + if main_full_att is None: + # Non-hybrid / unexpected mem_manager: nothing to remap. + return + + draft_idx = len(self.mtp_previous_draft_models) + draft_full_att_layers = getattr(mem_manager, "draft_full_att_layers", None) + if draft_full_att_layers is not None: + assert draft_idx < draft_full_att_layers, ( + f"draft_idx {draft_idx} out of range for draft_full_att_layers " + f"{draft_full_att_layers}; mem_manager not sized for this many MTP draft blocks" + ) + draft_kv_slot = main_full_att + draft_idx + layer_infer = self.layers_infer[0] + layer_infer.layer_num_ = draft_kv_slot * interval + logger.info( + f"Qwen3.5 MTP draft layer assigned dedicated full-attn KV slot {draft_kv_slot} " + f"(layer_num_={layer_infer.layer_num_}, interval={interval}, main_full_att={main_full_att})" + ) + return diff --git a/lightllm/models/qwen3_moe_mtp/model.py b/lightllm/models/qwen3_moe_mtp/model.py index 9f83832a7e..ec27b57070 100644 --- a/lightllm/models/qwen3_moe_mtp/model.py +++ b/lightllm/models/qwen3_moe_mtp/model.py @@ -1,13 +1,12 @@ -from typing import List +from lightllm.models.base_mtp_model import BaseMTPModel from lightllm.models.qwen3_moe.model import Qwen3MOEModel from lightllm.models.qwen3_moe_mtp.layer_weights.pre_and_post_layer_weight import Qwen3MOEMTPPreAndPostLayerWeight from lightllm.models.deepseek_mtp.layer_infer.pre_layer_infer import Deepseek3MTPPreLayerInfer from lightllm.models.qwen3_moe_mtp.layer_infer.transformer_layer_infer import Qwen3MOEMTPTransformerLayerInfer from lightllm.models.qwen3_moe_mtp.layer_weights.transformer_layer_weight import Qwen3MOEMTPTransformerLayerWeight -from lightllm.common.basemodel import TpPartBaseModel -class Qwen3MOEMTPModel(Qwen3MOEModel): +class Qwen3MOEMTPModel(BaseMTPModel, Qwen3MOEModel): pre_and_post_weight_class = Qwen3MOEMTPPreAndPostLayerWeight pre_layer_infer_class = Deepseek3MTPPreLayerInfer @@ -15,29 +14,6 @@ class Qwen3MOEMTPModel(Qwen3MOEModel): transformer_weight_class = Qwen3MOEMTPTransformerLayerWeight transformer_layer_infer_class = Qwen3MOEMTPTransformerLayerInfer - def __init__(self, kvargs: dict): - self._pre_init(kvargs) - super().__init__(kvargs) - return - - def _pre_init(self, kvargs: dict): - self.main_model: TpPartBaseModel = kvargs.pop("main_model") - self.mtp_previous_draft_models: List[TpPartBaseModel] = kvargs.pop("mtp_previous_draft_models") - return - - def _init_custom(self): - self._cos_cached = self.main_model._cos_cached - self._sin_cached = self.main_model._sin_cached - return - - def _init_req_manager(self): - self.req_manager = self.main_model.req_manager - return - - def _init_mem_manager(self): - self.mem_manager = self.main_model.mem_manager - return - def _init_weights(self, start_layer_index=None): assert start_layer_index is None mtp_index = len(self.mtp_previous_draft_models) diff --git a/lightllm/models/qwen3next/infer_struct.py b/lightllm/models/qwen3next/infer_struct.py index 0006a682f1..b486bc6040 100644 --- a/lightllm/models/qwen3next/infer_struct.py +++ b/lightllm/models/qwen3next/infer_struct.py @@ -1,6 +1,4 @@ -import torch from lightllm.models.llama.infer_struct import LlamaInferStateInfo -from lightllm.utils.envs_utils import get_env_start_args class Qwen3NextInferStateInfo(LlamaInferStateInfo): @@ -10,7 +8,7 @@ def __init__(self): def init_some_extra_state(self, model): super().init_some_extra_state(model) - self.b_att_seq_len = self.b_seq_len - mtp_step = get_env_start_args().mtp_step - self.b_buffer_idx = self.b_req_idx * (mtp_step + 1) + self.b_mtp_index + from lightllm.common.basemodel.mtp_verify_extra_state import init_mtp_verify_extra_state + + init_mtp_verify_extra_state(self) return diff --git a/lightllm/models/qwen3next/layer_infer/transformer_layer_infer.py b/lightllm/models/qwen3next/layer_infer/transformer_layer_infer.py index bb48bfe49c..76c273c0e7 100644 --- a/lightllm/models/qwen3next/layer_infer/transformer_layer_infer.py +++ b/lightllm/models/qwen3next/layer_infer/transformer_layer_infer.py @@ -45,7 +45,6 @@ def __init__(self, layer_num, network_config): return def _init_linear_layer_metadata(self, layer_num, network_config): - # Linear attention specific dimensions self.num_v_heads = network_config["linear_num_value_heads"] self.num_k_heads = network_config["linear_num_key_heads"] @@ -121,7 +120,6 @@ def _compute_shared_expert( def _moe_ffn_tp( self, input: torch.Tensor, infer_state: Qwen3NextInferStateInfo, layer_weight: Qwen3NextTransformerLayerWeight ): - shared_expert_out = self._compute_shared_expert(input, infer_state, layer_weight) hidden_states = input.view(-1, self.embed_dim_) @@ -254,6 +252,18 @@ def gdn_forward( if is_prefill: core_attn_out, z = self._gdn_prefill_wrapper_run(mixed_qkvzba, infer_state, layer_weight) + elif getattr(infer_state, "is_mtp_verify", False): + mixed_qkv, z, b, a = self._split_qkvzba(mixed_qkvzba) + conv_states, ssm_states = infer_state.req_manager.get_mamba_cache(self.layer_num_) + core_attn_out = self._gdn_verify_kernel( + mixed_qkv, + conv_states, + ssm_states, + a, + b, + infer_state, + layer_weight, + ) else: mixed_qkv, z, b, a = self._split_qkvzba(mixed_qkvzba) conv_states, ssm_states = infer_state.req_manager.get_mamba_cache(self.layer_num_) @@ -374,7 +384,7 @@ def _gdn_prefill_kernel( layer_weight.linear_conv1d.mm_param.weight, bias=layer_weight.linear_conv1d.bias, query_start_loc=infer_state.b1_cu_q_seq_len, - cache_indices=infer_state.b_buffer_idx, + cache_indices=infer_state.b_conv_buffer_idx, has_initial_state=infer_state.b_ready_cache_len > 0, conv_states=conv_states, activation=self.activation, @@ -419,7 +429,7 @@ def _gdn_decode_kernel( layer_weight.linear_conv1d.mm_param.weight, bias=layer_weight.linear_conv1d.bias, activation=self.activation, - conv_state_indices=infer_state.b_buffer_idx, + conv_state_indices=infer_state.b_conv_buffer_idx, ) # Recurrent processing with fused gating @@ -439,3 +449,51 @@ def _gdn_decode_kernel( b_raw=b, ) return core_attn_out + + def _gdn_verify_kernel( + self, + mixed_qkv: torch.Tensor, + conv_states: torch.Tensor, + ssm_states: torch.Tensor, + a: torch.Tensor, + b: torch.Tensor, + infer_state: Qwen3NextInferStateInfo, + layer_weight: Qwen3NextTransformerLayerWeight, + ): + from lightllm.models.qwen3next.triton_kernel.causal_conv1d_spec import ( + causal_conv1d_update as causal_conv1d_update_spec, + ) + + mixed_qkv = causal_conv1d_update_spec( + mixed_qkv, + conv_states, + layer_weight.linear_conv1d.mm_param.weight, + bias=layer_weight.linear_conv1d.bias, + activation=self.activation, + conv_state_indices=infer_state.b_conv_buffer_idx, + num_accepted_tokens=infer_state.b_num_accepted_tokens, + query_start_loc=infer_state.b_gdn_verify_cu_seqlens, + ) + + query, key, value = self._rearrange_mixed_qkv(mixed_qkv, decode=False) + assert infer_state.b_ssm_index_rows.dim() == 2, "SSM index rows must be 2D [N, S+1]" + # #8b: b_num_accepted_tokens >= 1 is guaranteed upstream (init sets accept_len=1; the + # offload/snapshot guards bound it to [1, mtp_step+1]). The old per-layer per-step .all() + # D2H sync stalled the GPU on the eager decode hot path; it is redundant here. + core_attn_out, _ = fused_recurrent_gated_delta_rule( + q=query, + k=key, + v=value, + initial_state=ssm_states, + inplace_final_state=True, + cu_seqlens=infer_state.b_gdn_verify_cu_seqlens.to(torch.long), + ssm_state_indices=infer_state.b_ssm_index_rows, + ssm_state_write_indices=infer_state.b_ssm_index_rows, + num_accepted_tokens=infer_state.b_num_accepted_tokens, + use_qk_l2norm_in_kernel=True, + A_log=layer_weight.linear_A_log.weight, + dt_bias=layer_weight.linear_dt_bias.weight, + a_raw=a, + b_raw=b, + ) + return core_attn_out diff --git a/lightllm/models/qwen3next/model.py b/lightllm/models/qwen3next/model.py index e3c51f3617..f0940ba0f8 100644 --- a/lightllm/models/qwen3next/model.py +++ b/lightllm/models/qwen3next/model.py @@ -16,14 +16,16 @@ from lightllm.common.kv_cache_mem_manager.qwen3next_mem_manager import Qwen3NextMemManager from lightllm.server.core.objs.start_args_type import StartArgs from lightllm.common.req_manager import ReqManagerForMamba -from lightllm.common.linear_att_cache_manager.config_objs import LinearAttCacheConfig +from lightllm.common.linear_att_cache_manager.config_objs import ( + LinearAttCacheConfig, + get_mtp_draft_full_att_layer_num, +) logger = init_logger(__name__) @ModelRegistry("qwen3_next") class Qwen3NextTpPartModel(Qwen3MOEModel): - # weight class pre_and_post_weight_class = Qwen3NextPreAndPostLayerWeight transformer_weight_class = Qwen3NextTransformerLayerWeight @@ -59,6 +61,7 @@ def _init_mem_manager(self): assert self.config["num_attention_heads"] % self.tp_world_size_ == 0 start_args: StartArgs = get_env_start_args() ssm_dtype_dict = {"bfloat16": torch.bfloat16, "float32": torch.float32} + draft_full_att_layers = get_mtp_draft_full_att_layer_num(start_args) self.linear_config = LinearAttCacheConfig( tp_world_size=self.tp_world_size_, full_att_all_num_kv_heads=self.config["num_key_value_heads"], @@ -76,17 +79,26 @@ def _init_mem_manager(self): ssm_state_dtype=ssm_dtype_dict[start_args.linear_att_ssm_data_type], full_attention_interval=self.config["full_attention_interval"], all_layer_num=self.config["n_layer"], + draft_full_att_layer_num=draft_full_att_layers, ) + main_full_att = self.linear_config.get_main_full_att_layer_num() + persisted_full_att = self.linear_config.get_persisted_full_att_layer_num() + self._main_full_att_layer_num = main_full_att + self._draft_full_att_layers = draft_full_att_layers + self.mem_manager = Qwen3NextMemManager( size=self.max_total_token_num, dtype=self.data_type, num_kv_heads=self.num_kv_heads, head_dim=self.config["head_dim"], - full_att_layer_num=self.linear_config.all_layer_num - self.linear_config.linear_layer_num, + full_att_layer_num=persisted_full_att, linear_config=self.linear_config, mem_fraction=self.mem_fraction, ) + self.mem_manager.main_full_att_layer_num = main_full_att + self.mem_manager.draft_full_att_layers = draft_full_att_layers + self.mem_manager.persisted_full_att_layer_num = persisted_full_att def _init_req_manager(self): create_max_seq_len = 0 diff --git a/lightllm/models/qwen3next/triton_kernel/causal_conv1d_spec.py b/lightllm/models/qwen3next/triton_kernel/causal_conv1d_spec.py new file mode 100644 index 0000000000..2f0e22fa3f --- /dev/null +++ b/lightllm/models/qwen3next/triton_kernel/causal_conv1d_spec.py @@ -0,0 +1,468 @@ +# Vendored from vLLM v0.14.1 +# source: vllm/model_executor/layers/mamba/ops/causal_conv1d.py +# commit: d7de043d55d1dd629554467e23874097e1c48993 +# Adapted for LightLLM: imports point at standard triton; the vLLM-specific +# block-table params (block_idx_last_scheduled_token, initial_state_idx, +# null_block_id) are dropped — LightLLM uses contiguous per-request slots. +# Supports spec-decode: writes per-position conv state to a single widened slot +# per request and reads from offset (num_accepted_tokens-1). +# +# Upstream copyright notice: +# SPDX-License-Identifier: Apache-2.0 +# SPDX-FileCopyrightText: Copyright contributors to the vLLM project +# Copyright (c) 2024, Tri Dao. +# Adapted from +# https://github.com/Dao-AILab/causal-conv1d/blob/main/causal_conv1d/causal_conv1d_interface.py +from typing import Optional + +import torch +import triton +import triton.language as tl + + +@triton.jit() +def _causal_conv1d_update_kernel( + # Pointers to matrices + x_ptr, # (batch, dim, seqlen) + w_ptr, # (dim, width) + bias_ptr, + conv_state_ptr, + conv_state_indices_ptr, + num_accepted_tokens_ptr, + query_start_loc_ptr, # (batch + 1) + o_ptr, # (batch, dim, seqlen) + # Matrix dimensions + batch: int, + dim: tl.constexpr, + seqlen: tl.constexpr, + state_len: tl.constexpr, + num_cache_lines: tl.constexpr, # added to support vLLM larger cache lines + # Strides + stride_x_seq: tl.constexpr, + stride_x_dim: tl.constexpr, + stride_x_token: tl.constexpr, + stride_w_dim: tl.constexpr, + stride_w_width: tl.constexpr, + stride_conv_state_seq: tl.constexpr, + stride_conv_state_dim: tl.constexpr, + stride_conv_state_tok: tl.constexpr, + stride_state_indices: tl.constexpr, + stride_o_seq: tl.constexpr, + stride_o_dim: tl.constexpr, + stride_o_token: tl.constexpr, + # others + pad_slot_id: tl.constexpr, + # Meta-parameters + HAS_BIAS: tl.constexpr, + KERNEL_WIDTH: tl.constexpr, + SILU_ACTIVATION: tl.constexpr, + IS_VARLEN: tl.constexpr, + IS_SPEC_DECODING: tl.constexpr, + NP2_STATELEN: tl.constexpr, + USE_PAD_SLOT: tl.constexpr, + BLOCK_N: tl.constexpr, +): + # ruff: noqa: E501 + idx_seq = tl.program_id(0) + if idx_seq >= batch: + return + + # [BLOCK_N,] elements along the feature-dimension (channel) + idx_feats = tl.program_id(1) * BLOCK_N + tl.arange(0, BLOCK_N) + + # LightLLM uses contiguous per-request slots, so the cache block for both + # the initial-state read and the final write is always conv_state_indices[idx_seq]. + conv_state_init = 0 + current_last_index = 0 + + # cache_idx + conv_states_input_coord = tl.load(conv_state_indices_ptr + idx_seq * stride_state_indices + conv_state_init).to( + tl.int64 + ) + + if USE_PAD_SLOT: # noqa + if conv_states_input_coord == pad_slot_id: + # not processing as this is not the actual sequence + return + + if IS_VARLEN: + query_start_index = tl.load(query_start_loc_ptr + idx_seq).to(tl.int64) + query_end_index = tl.load(query_start_loc_ptr + (idx_seq + 1)).to(tl.int64) + # revise state_len and seqlen + state_len = state_len - (seqlen - (query_end_index - query_start_index)) + seqlen = query_end_index - query_start_index + x_offset = query_start_index * stride_x_token + o_offset = query_start_index * stride_o_token + else: + query_start_index = idx_seq * seqlen + query_end_index = query_start_index + seqlen + x_offset = idx_seq * stride_x_seq + o_offset = idx_seq * stride_o_seq + + if query_start_index == query_end_index: + return + + if IS_SPEC_DECODING: + # The rolling of conv state: + # + # Before forward, the conv_state is: + # [history1, history2, ..., historyM]. + # + # After forward, the conv_state becomes: + # [history2, ..., historyM, draft1, draft2, ..., draftN]. + # + # After acceptance, it becomes: + # + # - accept 1 tokens: [history2, ..., historyM, draft1] + # - accept 2 tokens: [history3, ..., historyM, draft1, draft2] + # - and so on. + conv_state_token_offset = tl.load(num_accepted_tokens_ptr + idx_seq).to(tl.int64) - 1 + else: + conv_state_token_offset = 0 + + # STEP 1: READ init_state data + conv_states_base = ( + conv_state_ptr + (conv_states_input_coord * stride_conv_state_seq) + (idx_feats * stride_conv_state_dim) + ) + mask_w = idx_feats < dim + + prior_tokens = conv_states_base + conv_state_token_offset * stride_conv_state_tok + if KERNEL_WIDTH >= 2: + conv_states_ptrs = prior_tokens # [BLOCK_N] + col0 = tl.load(conv_states_ptrs, mask_w, 0.0) + if KERNEL_WIDTH >= 3: + conv_states_ptrs = prior_tokens + 1 * stride_conv_state_tok # [BLOCK_N] + col1 = tl.load(conv_states_ptrs, mask_w, 0.0) + if KERNEL_WIDTH >= 4: + conv_states_ptrs = prior_tokens + 2 * stride_conv_state_tok # [BLOCK_N] + col2 = tl.load(conv_states_ptrs, mask_w, 0.0) + if KERNEL_WIDTH >= 5: + conv_states_ptrs = prior_tokens + 3 * stride_conv_state_tok # [BLOCK_N] + col3 = tl.load(conv_states_ptrs, mask_w, 0.0) + if KERNEL_WIDTH >= 6: + conv_states_ptrs = prior_tokens + 4 * stride_conv_state_tok # [BLOCK_N] + col4 = tl.load(conv_states_ptrs, mask_w, 0.0) + + # STEP 2: assume state_len > seqlen + idx_tokens = tl.arange(0, NP2_STATELEN) # [BLOCK_M] + + # With speculative decoding, the conv_state updates works in a sliding + # window manner, at each forward pass, the tokens are shift by 1, so we + # load since idx_tokens + 1. + conv_state_ptrs_source = ( + conv_state_ptr + + (conv_states_input_coord * stride_conv_state_seq) + + conv_state_token_offset * stride_conv_state_tok + + (idx_feats * stride_conv_state_dim)[None, :] + + ((idx_tokens + (1 if IS_SPEC_DECODING else seqlen)) * stride_conv_state_tok)[:, None] + ) # [BLOCK_M, BLOCK_N] + mask = ( + (conv_states_input_coord < num_cache_lines) + & ((idx_tokens + seqlen) < state_len)[:, None] + & (idx_feats < dim)[None, :] + ) + conv_state = tl.load(conv_state_ptrs_source, mask, other=0.0) + + VAL = state_len - seqlen + x_base = x_ptr + x_offset + (idx_feats * stride_x_dim) # [BLOCK_N] + + x_ptrs = x_base[None, :] + ((idx_tokens - VAL) * stride_x_token)[:, None] # [BLOCK_M, BLOCK_N] + + mask_x = ( + (idx_tokens - VAL >= 0)[:, None] & (idx_tokens - VAL < seqlen)[:, None] & (idx_feats < dim)[None, :] + ) # token-index # token-index # feature-index + loaded_x = tl.load(x_ptrs, mask_x, 0.0) + tl.debug_barrier() + + new_conv_state = tl.where(mask, conv_state, loaded_x) + + # Write the updated state back. In LightLLM the read and write slots are the + # same contiguous per-request slot (current_last_index == conv_state_init == 0), + # so this resolves to the same conv_state_indices[idx_seq] used for the read. + conv_states_offset = tl.load(conv_state_indices_ptr + idx_seq * stride_state_indices + current_last_index).to( + tl.int64 + ) + conv_state_ptrs_target = ( + conv_state_ptr + + (conv_states_offset * stride_conv_state_seq) # Offset from seq + + (idx_feats * stride_conv_state_dim) + )[ + None, : + ] + ( # [BLOCK_N,] + idx_tokens * stride_conv_state_tok + )[ + :, None + ] + mask = (idx_tokens < state_len)[:, None] & (idx_feats < dim)[None, :] + tl.store(conv_state_ptrs_target, new_conv_state, mask) + + # STEP 3: init accumulator + if HAS_BIAS: + bias = bias_ptr + idx_feats + mask_bias = idx_feats < dim + acc_preload = tl.load(bias, mask=mask_bias, other=0.0).to(tl.float32) # [BLOCK_N] + else: + acc_preload = tl.zeros((BLOCK_N,), dtype=tl.float32) + + # STEP 4: + # PRE-LOAD WEIGHTS + # first kernel column, configured for weights to handle BLOCK_N features in range + w_base = w_ptr + (idx_feats * stride_w_dim) # [BLOCK_N,] + mask_w = idx_feats < dim + if KERNEL_WIDTH >= 2: + w_ptrs = w_base + (0 * stride_w_width) # [BLOCK_N] tensor + w_col0 = tl.load(w_ptrs, mask_w, other=0.0) + w_ptrs = w_base + (1 * stride_w_width) # [BLOCK_N] tensor + w_col1 = tl.load(w_ptrs, mask_w, other=0.0) + if KERNEL_WIDTH >= 3: + w_ptrs = w_base + (2 * stride_w_width) # [BLOCK_N] tensor + w_col2 = tl.load(w_ptrs, mask_w, other=0.0) + if KERNEL_WIDTH >= 4: + w_ptrs = w_base + (3 * stride_w_width) # [BLOCK_N] tensor + w_col3 = tl.load(w_ptrs, mask_w, other=0.0) + if KERNEL_WIDTH >= 5: + w_ptrs = w_base + (4 * stride_w_width) # [BLOCK_N] tensor + w_col4 = tl.load(w_ptrs, mask_w, other=0.0) + if KERNEL_WIDTH >= 6: + w_ptrs = w_base + (5 * stride_w_width) # [BLOCK_N] tensor + w_col5 = tl.load(w_ptrs, mask_w, other=0.0) + + x_base_1d = x_base # starting of chunk [BLOCK_N] + mask_x_1d = idx_feats < dim + + # STEP 5: compute each token + for idx_token in tl.range(seqlen): + acc = acc_preload + + matrix_w = w_col0 + matrix_x = col0 + for j in tl.static_range(KERNEL_WIDTH): + if KERNEL_WIDTH == 2: + if j == 1: # KERNEL_WIDTH-1: + matrix_w = w_col1 + x_ptrs_1d = x_base_1d + idx_token * stride_x_token # [BLOCK_N] + matrix_x = tl.load(x_ptrs_1d, mask=mask_x_1d) + elif KERNEL_WIDTH == 3: + if j == 1: + matrix_w = w_col1 + matrix_x = col1 + elif j == 2: + matrix_w = w_col2 + x_ptrs_1d = x_base_1d + idx_token * stride_x_token # [BLOCK_N] + matrix_x = tl.load(x_ptrs_1d, mask=mask_x_1d) + elif KERNEL_WIDTH == 4: + if j == 1: + matrix_w = w_col1 + matrix_x = col1 + elif j == 2: + matrix_w = w_col2 + matrix_x = col2 + elif j == 3: + matrix_w = w_col3 + x_ptrs_1d = x_base_1d + idx_token * stride_x_token # [BLOCK_N] + matrix_x = tl.load(x_ptrs_1d, mask=mask_x_1d) + elif KERNEL_WIDTH == 5: + if j == 1: + matrix_w = w_col1 + matrix_x = col1 + elif j == 2: + matrix_w = w_col2 + matrix_x = col2 + elif j == 3: + matrix_w = w_col3 + matrix_x = col3 + elif j == 4: + matrix_w = w_col4 + x_ptrs_1d = x_base_1d + idx_token * stride_x_token # [BLOCK_N] + matrix_x = tl.load(x_ptrs_1d, mask=mask_x_1d) + elif KERNEL_WIDTH == 6: + if j == 1: + matrix_w = w_col1 + matrix_x = col1 + elif j == 2: + matrix_w = w_col2 + matrix_x = col2 + elif j == 3: + matrix_w = w_col3 + matrix_x = col3 + elif j == 4: + matrix_w = w_col4 + matrix_x = col4 + elif j == 5: + matrix_w = w_col5 + x_ptrs_1d = x_base_1d + idx_token * stride_x_token # [BLOCK_N] + matrix_x = tl.load(x_ptrs_1d, mask=mask_x_1d) + + acc += matrix_x * matrix_w # [BLOCK_N] + + if KERNEL_WIDTH == 2: + col0 = matrix_x + elif KERNEL_WIDTH == 3: + col0 = col1 + col1 = matrix_x + elif KERNEL_WIDTH == 4: + col0 = col1 + col1 = col2 + col2 = matrix_x + elif KERNEL_WIDTH == 5: + col0 = col1 + col1 = col2 + col2 = col3 + col3 = matrix_x + elif KERNEL_WIDTH == 6: + col0 = col1 + col1 = col2 + col2 = col3 + col3 = col4 + col4 = matrix_x + + if SILU_ACTIVATION: + acc = acc / (1 + tl.exp(-acc)) + mask_1d = (idx_token < seqlen) & (idx_feats < dim) # token-index # feature-index + o_ptrs = o_ptr + o_offset + idx_token * stride_o_token + (idx_feats * stride_o_dim) + + tl.store(o_ptrs, acc, mask=mask_1d) + + +def causal_conv1d_update( + x: torch.Tensor, + conv_state: torch.Tensor, + weight: torch.Tensor, + bias: Optional[torch.Tensor] = None, + activation: Optional[str] = None, + cache_seqlens: Optional[torch.Tensor] = None, + conv_state_indices: Optional[torch.Tensor] = None, + num_accepted_tokens: Optional[torch.Tensor] = None, + query_start_loc: Optional[torch.Tensor] = None, + pad_slot_id: int = -1, +): + """Spec-decode capable conv1d update. When num_accepted_tokens/query_start_loc + are None it must behave like a single-token decode update. x may be (batch, dim) + single-token or (num_tokens, dim) flattened varlen with query_start_loc grouping + each request's S+1 candidates. conv_state is (num_slots, dim, state_len) with + state_len = (width-1)+S widened. Read offset = num_accepted_tokens-1; writes to + the same slot. + + Args: + x: input tensor of shape ``(batch, dim)`` (single-token decode), + ``(batch, dim, seqlen)`` (single/multi token), or ``(num_tokens, dim)`` + flattened varlen grouped by ``query_start_loc``. + conv_state: ``(num_slots, dim, state_len)`` with ``state_len >= width - 1``. + For spec decode the slot is widened to ``(width - 1) + S`` where ``S`` is + the number of speculative tokens (so ``seqlen == S + 1``). + weight: depthwise filter of shape ``(dim, width)``. + bias: optional ``(dim,)`` bias. + activation: ``None``, ``"silu"`` or ``"swish"``. + cache_seqlens: accepted for call-compatibility with the non-spec wrapper; + unused here. + conv_state_indices: ``(batch,)`` int32 mapping each request to its conv_state + slot. Required when ``query_start_loc`` is given. + num_accepted_tokens: ``(batch,)`` int32. When not None the conv_state read + offset for each request is ``num_accepted_tokens - 1`` (sliding window + spec-decode update). + query_start_loc: ``(batch + 1,)`` int32 varlen cumulative token offsets; when + None the call is a plain single-/multi-token decode update. + pad_slot_id: slot id that marks padded entries to skip. + + Returns: + Output tensor with the same shape as ``x`` (the kernel overwrites ``x`` in + place), one conv output per input token. + """ + if activation is not None: + assert activation in ["silu", "swish"] + + original_x_dtype = x.dtype + x = x.to(conv_state.dtype) + unsqueeze = query_start_loc is None and x.dim() == 2 + if unsqueeze: + # make it (batch, dim, seqlen) with seqlen == 1 + x = x.unsqueeze(-1) + if query_start_loc is None: + batch, dim, seqlen = x.shape + else: + assert conv_state_indices is not None + batch = conv_state_indices.size(0) + dim = x.size(1) + # The MTP verify layout is uniform (mtp_step+1) tokens per request, so seqlen is + # structurally x.size(0) // batch. Compute it without a D2H sync on query_start_loc on + # BOTH the capture and eager paths (#8a) — the eager .item() ran once per GDN layer per + # decode step. .item() is also illegal during CUDA-graph capture. + assert x.size(0) % batch == 0, "varlen conv update expects a uniform per-request length" + seqlen = x.size(0) // batch + _, width = weight.shape + # conv_state: (num_slots, dim, state_len), where state_len >= width - 1 + num_cache_lines, _, state_len = conv_state.size() + + # adopt the strategy in vLLM that overwrites 'x' directly, rather than creating a new tensor 'o' + out = x + stride_w_dim, stride_w_width = weight.stride() + + if query_start_loc is None: + # X (batch, dim, seqlen) + stride_x_seq, stride_x_dim, stride_x_token = x.stride() + stride_o_seq, stride_o_dim, stride_o_token = out.stride() + else: + # X (num_tokens, dim) + stride_x_token, stride_x_dim = x.stride() + stride_x_seq = 0 + stride_o_token, stride_o_dim = out.stride() + stride_o_seq = 0 + + stride_istate_seq, stride_istate_dim, stride_istate_token = conv_state.stride() + stride_state_indices = conv_state_indices.stride(0) if conv_state_indices is not None else 0 + if num_accepted_tokens is not None: + state_len = width - 1 + (seqlen - 1) # effective state_len needed + else: + state_len = width - 1 + np2_statelen = triton.next_power_of_2(state_len) + + def grid(META): + return ( + batch, + triton.cdiv(dim, META["BLOCK_N"]), + ) + + _causal_conv1d_update_kernel[grid]( + # Pointers to matrices + x, + weight, + bias, + conv_state, + conv_state_indices, + num_accepted_tokens, + query_start_loc, + out, + # Matrix dimensions + batch, + dim, + seqlen, + state_len, + num_cache_lines, + # stride + stride_x_seq, + stride_x_dim, + stride_x_token, + stride_w_dim, + stride_w_width, + stride_istate_seq, + stride_istate_dim, + stride_istate_token, + stride_state_indices, + stride_o_seq, + stride_o_dim, + stride_o_token, + # others + pad_slot_id, + # META + HAS_BIAS=bias is not None, + KERNEL_WIDTH=width, + SILU_ACTIVATION=activation in ["silu", "swish"], + IS_VARLEN=query_start_loc is not None, + IS_SPEC_DECODING=num_accepted_tokens is not None, + NP2_STATELEN=np2_statelen, + USE_PAD_SLOT=pad_slot_id is not None, + BLOCK_N=256, + ) + if unsqueeze: + out = out.squeeze(-1) + return out.to(original_x_dtype) diff --git a/lightllm/server/router/model_infer/infer_batch.py b/lightllm/server/router/model_infer/infer_batch.py index f0ec69b2c1..72df815183 100644 --- a/lightllm/server/router/model_infer/infer_batch.py +++ b/lightllm/server/router/model_infer/infer_batch.py @@ -357,6 +357,11 @@ def copy_linear_att_state_to_cache_buffer(self, b_req_idx: torch.Tensor, reqs: L if not self.is_linear_att_mixed_model: return + # 当 dynamic prompt cache 被禁用时 radix_cache 为 None,没有大页/小页缓冲可写, + # 线性层状态仅存于 req_manager 的 GPU buffer 即可,直接跳过跨请求缓存拷贝。 + if self.radix_cache is None: + return + # 大页对应的 linear att 的拷贝 big_page_token_num = self.args.linear_att_hash_page_size * self.args.linear_att_page_block_num big_page_buffer_ids = [] @@ -377,6 +382,10 @@ def copy_linear_att_state_to_cache_buffer(self, b_req_idx: torch.Tensor, reqs: L from lightllm.common.basemodel.triton_kernel.linear_att_copy import copy_linear_att_state_to_kv_buffer + b_num_accepted_tokens = torch.tensor( + [req.mtp_accept_len for req in reqs], dtype=torch.int32, requires_grad=False, device="cpu" + ).cuda(non_blocking=True) + copy_linear_att_state_to_kv_buffer( b_req_idx=b_req_idx, big_page_buffer_ids=big_page_buffer_ids, @@ -385,6 +394,7 @@ def copy_linear_att_state_to_cache_buffer(self, b_req_idx: torch.Tensor, reqs: L cpu_kv_conv_state=self.radix_cache.linear_att_big_page_buffers.conv_state_cache.buffer, cpu_kv_ssm_state=self.radix_cache.linear_att_big_page_buffers.ssm_state_cache.buffer, mtp_step=self.args.mtp_step, + b_num_accepted_tokens=b_num_accepted_tokens, ) assert not self.args.disable_chunked_prefill, "chunked prefill mode must be enabled for linear att mixed model" @@ -400,9 +410,18 @@ def copy_linear_att_state_to_cache_buffer(self, b_req_idx: torch.Tensor, reqs: L self.radix_cache.linear_att_small_page_buffers.alloc_one_state_cache() ) if req.tail_linear_att_small_page_buffer_id is not None: - src_buffer_idx = req.req_idx * (self.args.mtp_step + 1) - gpu_conv_state = self.req_manager.req_to_conv_state.buffer[:, src_buffer_idx, ...] - gpu_ssm_state = self.req_manager.req_to_ssm_state.buffer[:, src_buffer_idx, ...] + assert 1 <= req.mtp_accept_len <= self.args.mtp_step + 1, ( + f"mtp_accept_len={req.mtp_accept_len} out of range " + f"[1, {self.args.mtp_step + 1}]; would slice past the widened conv slot" + ) + canonical_off = req.mtp_accept_len - 1 + conv_src_idx = req.req_idx + ssm_src_idx = req.req_idx * (self.args.mtp_step + 1) + canonical_off + narrow_w = self.req_manager.linear_config.get_persisted_conv_state_shape()[-1] + gpu_conv_state = self.req_manager.req_to_conv_state.buffer[ + :, conv_src_idx, ..., canonical_off : canonical_off + narrow_w + ] + gpu_ssm_state = self.req_manager.req_to_ssm_state.buffer[:, ssm_src_idx, ...] dst_buffer_idx = req.tail_linear_att_small_page_buffer_id dst_conv_state, dst_ssm_state = self.radix_cache.linear_att_small_page_buffers.get_state_cache( @@ -558,6 +577,8 @@ def __init__( else: self.decode_need_token_num = self._normal_decode_need_token_num + self.mtp_accept_len: int = 1 + if g_infer_context.is_linear_att_mixed_model: self.get_chuncked_input_token_len = self.get_chuncked_input_token_len_for_linear_att self.get_chuncked_input_token_ids = self.get_chuncked_input_token_ids_for_linear_att diff --git a/lightllm/server/router/model_infer/mode_backend/base_backend.py b/lightllm/server/router/model_infer/mode_backend/base_backend.py index 0220dc87fb..768c9126de 100644 --- a/lightllm/server/router/model_infer/mode_backend/base_backend.py +++ b/lightllm/server/router/model_infer/mode_backend/base_backend.py @@ -1,4 +1,5 @@ import os +import copy import numpy as np import torch import time @@ -42,10 +43,6 @@ ) from lightllm.server.core.objs.shm_objs_io_buffer import ShmObjsIOBuffer from lightllm.server.router.model_infer.mode_backend.overlap_events import OverlapEventManager, OverlapEventPack -from lightllm.models.deepseek_mtp.model import Deepseek3MTPModel -from lightllm.models.qwen3_moe_mtp.model import Qwen3MOEMTPModel -from lightllm.models.mistral_mtp.model import MistralMTPModel -from lightllm.models.glm4_moe_lite_mtp.model import Glm4MoeLiteMTPModel from lightllm.server.router.model_infer.mode_backend.generic_post_process import sample from lightllm.common.basemodel.triton_kernel.gather_token_id import scatter_token from lightllm.server.pd_io_struct import NIXLChunckedTransTaskRet @@ -345,22 +342,11 @@ def init_mtp_draft_model(self, main_kvargs: dict): "mtp_previous_draft_models": self.draft_models.copy(), } - # Select MTP model class based on model type + # Select MTP model class based on model type (single source of truth: #10). + from lightllm.server.router.model_infer.mode_backend.mtp_model_factory import create_mtp_draft_model + model_type = mtp_model_cfg.get("model_type", "") - if model_type == "deepseek_v3": - assert self.args.mtp_mode in ["vanilla_with_att", "eagle_with_att"] - self.draft_models.append(Deepseek3MTPModel(mtp_model_kvargs)) - elif model_type == "qwen3_moe": - assert self.args.mtp_mode in ["vanilla_no_att", "eagle_no_att"] - self.draft_models.append(Qwen3MOEMTPModel(mtp_model_kvargs)) - elif model_type == "mistral": - assert self.args.mtp_mode in ["vanilla_no_att", "eagle_no_att"] - self.draft_models.append(MistralMTPModel(mtp_model_kvargs)) - elif mtp_model_cfg["model_type"] == "glm4_moe_lite": - assert self.args.mtp_mode in ["vanilla_with_att", "eagle_with_att"] - self.draft_models.append(Glm4MoeLiteMTPModel(mtp_model_kvargs)) - else: - raise ValueError(f"Unsupported MTP model type: {model_type}") + self.draft_models.append(create_mtp_draft_model(model_type, self.args.mtp_mode, mtp_model_kvargs)) self.logger.info(f"loaded mtp model class {self.draft_models[i].__class__}") return @@ -604,7 +590,6 @@ def _get_classed_reqs( can_alloc_token_num = g_infer_context.get_can_alloc_token_num() for req_obj in ready_reqs: - if req_obj.filter_mark: finished_reqs.append(req_obj) continue @@ -785,11 +770,69 @@ def _verify_mtp_v2( ) return mtp_accept_len, accepted_index + def _build_eagle_accepted_draft_input( + self, + main_model_input: ModelInput, + main_model_output: ModelOutput, + next_token_ids: torch.Tensor, + mtp_accept_len: torch.Tensor, + b_req_mtp_start_loc: torch.Tensor, + ): + accepted_row_idx = b_req_mtp_start_loc + mtp_accept_len - 1 + accepted_row_idx_long = accepted_row_idx.long() + + draft_model_input = copy.copy(main_model_input) + draft_model_input.batch_size = accepted_row_idx.shape[0] + draft_model_input.total_token_num = draft_model_input.batch_size * main_model_input.max_kv_seq_len + draft_model_input.input_ids = next_token_ids.index_select(0, accepted_row_idx_long) + draft_model_input.mtp_draft_input_hiddens = main_model_output.mtp_main_output_hiddens.index_select( + 0, accepted_row_idx_long + ) + draft_model_input.b_req_idx = main_model_input.b_req_idx.index_select(0, accepted_row_idx_long) + draft_model_input.b_mtp_index = main_model_input.b_mtp_index.index_select(0, accepted_row_idx_long) + draft_model_input.b_seq_len = main_model_input.b_seq_len.index_select(0, accepted_row_idx_long) + draft_model_input.b_num_accepted_tokens = None + if main_model_input.mem_indexes is not None: + draft_model_input.mem_indexes = main_model_input.mem_indexes.index_select(0, accepted_row_idx_long) + draft_model_input.mem_indexes_cpu = None + if main_model_input.b_shared_seq_len is not None: + draft_model_input.b_shared_seq_len = main_model_input.b_shared_seq_len.index_select( + 0, accepted_row_idx_long + ) + if main_model_input.b_mark_shared_group is not None: + draft_model_input.b_mark_shared_group = main_model_input.b_mark_shared_group.index_select( + 0, accepted_row_idx_long + ) + + if accepted_row_idx.device.type == "cpu": + selected_rows = accepted_row_idx.tolist() + draft_model_input.multimodal_params = [main_model_input.multimodal_params[i] for i in selected_rows] + else: + draft_model_input.multimodal_params = [ + {"images": [], "audios": []} for _ in range(draft_model_input.batch_size) + ] + + accepted_next_token_ids = draft_model_input.input_ids + accepted_req_idx = draft_model_input.b_req_idx + return draft_model_input, accepted_next_token_ids, accepted_req_idx + + def _scatter_accepted_next_token_ids(self, accepted_req_idx: torch.Tensor, all_next_token_ids: torch.Tensor): + req_to_next_token_ids = self.model.req_manager.req_sampling_params_manager.req_to_next_token_ids + width = all_next_token_ids.shape[1] + req_to_next_token_ids[:, :width].index_copy_( + 0, + accepted_req_idx.long(), + all_next_token_ids.to(dtype=req_to_next_token_ids.dtype), + ) + return + def _update_mtp_accept_ratio( self, decode_reqs: List[InferReq], mtp_accept_len_cpu: torch.Tensor, ): + # Master-only accept-ratio statistics. Unlike the phase-2 mtp_accept_len commit + # (inlined in decode_mtp) this only feeds metrics, so it may stay in phase 3. if self.is_master_in_dp: for req, accept_len in zip(decode_reqs, mtp_accept_len_cpu): req.update_mtp_accepted_token_num(accept_token_num=accept_len - 1) @@ -797,8 +840,9 @@ def _update_mtp_accept_ratio( def _gen_argmax_token_ids(self, model_output: ModelOutput): logits = model_output.logits - probs = torch.softmax(logits, dim=-1) - draft_next_token_ids_gpu = torch.argmax(probs, dim=-1) + # softmax is strictly monotonic, so argmax(softmax(logits)) == argmax(logits); + # skip the softmax to shorten the per-step MTP draft critical chain (need-to-fix #16). + draft_next_token_ids_gpu = torch.argmax(logits, dim=-1) return draft_next_token_ids_gpu def _sample_and_scatter_token( @@ -811,7 +855,6 @@ def _sample_and_scatter_token( b_prefill_has_output_cpu: torch.Tensor = None, mask_func: Optional[Callable] = None, ): - if mask_func is not None: assert len(run_reqs) == logits.shape[0] mask_func(run_reqs, logits) diff --git a/lightllm/server/router/model_infer/mode_backend/chunked_prefill/impl.py b/lightllm/server/router/model_infer/mode_backend/chunked_prefill/impl.py index 60045fab6c..a716e5663f 100644 --- a/lightllm/server/router/model_infer/mode_backend/chunked_prefill/impl.py +++ b/lightllm/server/router/model_infer/mode_backend/chunked_prefill/impl.py @@ -1,5 +1,6 @@ import torch import time +import copy from typing import List, Optional, Callable, Dict, Any from queue import Queue from lightllm.server.router.model_infer.mode_backend.base_backend import ModeBackend @@ -241,17 +242,23 @@ def decode_mtp( """ model_input, run_reqs = prepare_decode_inputs(decode_reqs) + if self.mtp_step > 0: + accept_lens = [req.mtp_accept_len for req in decode_reqs] + model_input.b_num_accepted_tokens = g_pin_mem_manager.gen_from_list( + key="b_num_accepted_tokens", + data=accept_lens, + dtype=torch.int32, + ) + with torch.cuda.stream(g_infer_context.get_overlap_stream()): - b_mtp_index_cpu = model_input.b_mtp_index model_output = self.model.forward(model_input) next_token_ids, next_token_logprobs = sample(model_output.logits, run_reqs, self.eos_id) - # verify the next_token_ids - b_req_mtp_start_loc = [index for index, mtp_index in enumerate(b_mtp_index_cpu) if mtp_index == 0] - b_req_mtp_start_loc = g_pin_mem_manager.gen_from_list( - key="b_req_mtp_start_loc", - data=b_req_mtp_start_loc, - dtype=torch.int32, - ).cuda(non_blocking=True) + # verify the next_token_ids. The chunked decode batch is the contiguous + # (mtp_step+1)-expanded layout, so request starts are structurally + # arange(n_real)*(mtp_step+1). Compute on device instead of a per-step Python + # list-comp + pinned pack + H2D (#22). + n_real = model_input.batch_size // (self.mtp_step + 1) + b_req_mtp_start_loc = torch.arange(n_real, dtype=torch.int32, device="cuda") * (self.mtp_step + 1) mtp_accept_len, accepted_index = self._verify_mtp_v2( new_next_token_ids=next_token_ids, @@ -293,6 +300,8 @@ def decode_mtp( # 第二阶段 event_pack.notify_post_handle_and_wait_pre_post_handle() verify_event.synchronize() + for req, accept_len in zip(decode_reqs, mtp_accept_len_cpu): + req.mtp_accept_len = int(accept_len) verify_ok_reqs = [run_reqs[i] for i in range(len(run_reqs)) if accepted_index_cpu[i] == 1] update_packs = self._pre_post_handle(verify_ok_reqs, is_chuncked_mode=False) @@ -347,15 +356,19 @@ def _draft_decode_vanilla( mtp_accept_len: torch.Tensor, b_req_mtp_start_loc: torch.Tensor, ): - # share some inference info with the main model - draft_model_input = main_model_input + # share some inference info with the main model. copy.copy 后清空 b_num_accepted_tokens, + # 使 draft (MTP) forward 走普通 decode 布局 (bs, False);否则会沿用主模型 decode_mtp 设置的 + # verify 布局,命中 MTP draft 模型从未捕获的 cudagraph key (bs, True) -> KeyError + # (cudagraph 关闭时则会在扁平的 draft batch 上误用 S+1 分组的 verify attention)。 + # 镜像 eagle 路径 _build_eagle_accepted_draft_input 中清空 b_num_accepted_tokens 的处理。 + draft_model_input = copy.copy(main_model_input) + draft_model_input.b_num_accepted_tokens = None draft_model_output = main_model_output draft_next_token_ids = next_token_ids all_next_token_ids = [] all_next_token_ids.append(next_token_ids) # process the draft model output for draft_model_idx in range(self.mtp_step): - draft_model_input.input_ids = draft_next_token_ids draft_model_input.mtp_draft_input_hiddens = draft_model_output.mtp_main_output_hiddens # spec decode: MTP @@ -382,8 +395,7 @@ def _draft_decode_eagle( mtp_accept_len: torch.Tensor, b_req_mtp_start_loc: torch.Tensor, ): - batch_size = main_model_input.batch_size - num_reqs = batch_size // (self.mtp_step + 1) + num_reqs = b_req_mtp_start_loc.shape[0] g_infer_state_lock.acquire() if g_infer_context.radix_cache is not None: g_infer_context.radix_cache.free_radix_cache_to_get_enough_token(num_reqs * self.mtp_step) @@ -391,37 +403,45 @@ def _draft_decode_eagle( g_infer_state_lock.release() eagle_mem_indexes = eagle_mem_indexes_cpu.cuda(non_blocking=True) - # share some inference info with the main model - draft_model_input = main_model_input + ( + draft_model_input, + draft_next_token_ids, + accepted_req_idx, + ) = self._build_eagle_accepted_draft_input( + main_model_input=main_model_input, + main_model_output=main_model_output, + next_token_ids=next_token_ids, + mtp_accept_len=mtp_accept_len, + b_req_mtp_start_loc=b_req_mtp_start_loc, + ) draft_model_output = main_model_output - draft_next_token_ids = next_token_ids all_next_token_ids = [] - all_next_token_ids.append(next_token_ids) - # process the draft model output - for _step in range(self.mtp_step): + all_next_token_ids.append(draft_next_token_ids) + + mtp_size = self.mtp_step + 1 + main_mem_indexes = main_model_input.mem_indexes.view(num_reqs, mtp_size) + eagle_mem_indexes_by_req = eagle_mem_indexes.view(self.mtp_step, num_reqs).transpose(0, 1).contiguous() + mem_index_plan = torch.cat([main_mem_indexes, eagle_mem_indexes_by_req], dim=1) + accepted_offsets = mtp_accept_len.long() - 1 + req_offsets = torch.arange(num_reqs, dtype=torch.long, device=mtp_accept_len.device) + for _step in range(self.mtp_step): draft_model_input.input_ids = draft_next_token_ids - draft_model_input.mtp_draft_input_hiddens = draft_model_output.mtp_main_output_hiddens + if _step > 0: + draft_model_input.mtp_draft_input_hiddens = draft_model_output.mtp_main_output_hiddens + draft_model_input.mem_indexes = mem_index_plan[req_offsets, accepted_offsets + _step] # spec decode: MTP draft_model_idx = _step % self.num_mtp_models draft_model_output: ModelOutput = self.draft_models[draft_model_idx].forward(draft_model_input) draft_next_token_ids = self._gen_argmax_token_ids(draft_model_output) draft_model_input.b_seq_len += 1 draft_model_input.max_kv_seq_len += 1 - eagle_mem_indexes_i = eagle_mem_indexes[_step * num_reqs : (_step + 1) * num_reqs] - draft_model_input.mem_indexes = torch.cat( - [draft_model_input.mem_indexes.view(-1, self.mtp_step + 1)[:, 1:], eagle_mem_indexes_i.view(-1, 1)], - dim=1, - ).view(-1) all_next_token_ids.append(draft_next_token_ids) all_next_token_ids = torch.stack(all_next_token_ids, dim=1) # [batch_size, mtp_step + 1] - mtp_scatter_next_token_ids( - req_to_next_token_ids=self.model.req_manager.req_sampling_params_manager.req_to_next_token_ids, - b_req_mtp_start_loc=b_req_mtp_start_loc, + self._scatter_accepted_next_token_ids( + accepted_req_idx=accepted_req_idx, all_next_token_ids=all_next_token_ids, - b_req_idx=main_model_input.b_req_idx, - mtp_accept_len=mtp_accept_len, ) return eagle_mem_indexes_cpu diff --git a/lightllm/server/router/model_infer/mode_backend/dp_backend/impl.py b/lightllm/server/router/model_infer/mode_backend/dp_backend/impl.py index f1896b201a..507db65d99 100644 --- a/lightllm/server/router/model_infer/mode_backend/dp_backend/impl.py +++ b/lightllm/server/router/model_infer/mode_backend/dp_backend/impl.py @@ -1,3 +1,4 @@ +import copy import torch import time import torch.nn.functional as F @@ -267,7 +268,6 @@ def prefill_overlap(self, event_pack: OverlapEventPack, prefill_reqs: List[Infer b_req_idx = torch.cat((model_input0.b_req_idx[0:req_num0], model_input1.b_req_idx[0:req_num1]), dim=0) if (req_num0 + req_num1) > 0: - _, next_token_ids_cpu, next_token_logprobs_cpu = self._sample_and_scatter_token( logits=logits, b_req_idx=b_req_idx, @@ -409,7 +409,6 @@ def prefill_mtp(self, event_pack: OverlapEventPack, prefill_reqs: List[InferReq] sync_event.record() if req_num > 0: - # 第二阶段 event_pack.notify_post_handle_and_wait_pre_post_handle() update_packs = self._pre_post_handle(run_reqs, is_chuncked_mode=not self.disable_chunked_prefill) @@ -436,10 +435,22 @@ def prefill_mtp(self, event_pack: OverlapEventPack, prefill_reqs: List[InferReq] return def decode_mtp(self, event_pack: OverlapEventPack, decode_reqs: List[InferReq]): - model_input, run_reqs, _ = padded_prepare_decode_inputs(decode_reqs) + model_input, run_reqs, padded_req_num = padded_prepare_decode_inputs(decode_reqs) b_mtp_index_cpu = model_input.b_mtp_index req_num = len(run_reqs) + if self.mtp_step > 0: + # 标记 verify decode 布局:每个 req 一个 accept 数量(padding 出来的 fake req 记为 1)。 + # 不设置 b_num_accepted_tokens 会让主模型的 verify forward 走非 verify 的 GDN/FA3 布局, + # 并命中 hybrid 主模型从未捕获的 cudagraph key (bs, False) -> KeyError。 + # 与 chunked_prefill/impl.py 的 decode_mtp 保持一致。 + accept_lens = [req.mtp_accept_len for req in decode_reqs] + [1] * padded_req_num + model_input.b_num_accepted_tokens = g_pin_mem_manager.gen_from_list( + key="b_num_accepted_tokens", + data=accept_lens, + dtype=torch.int32, + ) + with torch.cuda.stream(g_infer_context.get_overlap_stream()): model_output = self.model.forward(model_input) mtp_accept_len, b_req_mtp_start_loc, next_token_ids = None, None, None @@ -500,6 +511,11 @@ def decode_mtp(self, event_pack: OverlapEventPack, decode_reqs: List[InferReq]): # 第二阶段 event_pack.notify_post_handle_and_wait_pre_post_handle() verify_event.synchronize() + # 写回每个 req 的本步 accept 数量,供下一步 verify 经 b_num_accepted_tokens 传入 + # GDN/linear-att verify kernel(据此提交 conv/ssm 递归状态的正确偏移)。chunked 路径 + # 在 chunked_prefill/impl.py 同样写回;dp 缺失会让状态停留在 accept=1 -> 状态错乱、精度崩塌。 + for req, accept_len in zip(decode_reqs, mtp_accept_len_cpu): + req.mtp_accept_len = int(accept_len) verify_ok_reqs = [run_reqs[i] for i in range(len(run_reqs)) if accepted_index_cpu[i] == 1] update_packs = self._pre_post_handle(verify_ok_reqs, is_chuncked_mode=False) @@ -542,8 +558,11 @@ def _draft_decode_vanilla( req_num: int, ): all_next_token_ids = [] - # share some inference info with the main model - draft_model_input = model_input + # share some inference info with the main model. copy.copy 后清空 b_num_accepted_tokens, + # 使 draft (MTP) forward 走普通 decode 布局 (bs, False);否则会沿用主模型设置的 verify 布局, + # 命中 MTP draft 模型从未捕获的 cudagraph key (bs, True) -> KeyError。 + draft_model_input = copy.copy(model_input) + draft_model_input.b_num_accepted_tokens = None draft_model_output = model_output draft_next_token_ids_gpu = torch.zeros((model_input.batch_size), dtype=torch.int64, device="cuda") if req_num > 0: @@ -553,7 +572,6 @@ def _draft_decode_vanilla( # process the draft model output for draft_model_idx in range(self.mtp_step): - draft_model_input.input_ids = draft_next_token_ids_gpu draft_model_input.mtp_draft_input_hiddens = draft_model_output.mtp_main_output_hiddens # spec decode: MTP @@ -573,6 +591,64 @@ def _draft_decode_vanilla( ) return None + def _build_padding_draft_input(self, model_input: ModelInput, model_output: ModelOutput, common_req_num: int): + """ + 构造一个纯 padding 的 draft 输入,用于本 rank 没有真实 decode 请求 (real_req_num == 0) + 但其它 dp rank 仍有请求、需要本 rank 同步参与 mtp_step 次 draft forward 的集合通信的场景。 + + 从已 padding 的 main model_input 中按 (mtp_step+1) 分组取每组首行 (mtp_index==0) 即可, + 这些行均为 HOLD_REQUEST_ID / HOLD_TOKEN_MEMINDEX 的占位行。step0 的 hiddens 沿用主模型 + 对应占位行的 mtp_main_output_hiddens, 与原 DP 实现 (step0 使用 model_output.mtp_main_output_hiddens) + 保持一致, 避免 None 触发 draft forward 崩溃。 + """ + mtp_size = self.mtp_step + 1 + select_idx = torch.arange(common_req_num, dtype=torch.long, device=model_input.b_req_idx.device) * mtp_size + + draft_model_input = copy.copy(model_input) + draft_model_input.batch_size = common_req_num + draft_model_input.total_token_num = common_req_num * model_input.max_kv_seq_len + draft_model_input.b_num_accepted_tokens = None + draft_model_input.b_req_idx = model_input.b_req_idx.index_select(0, select_idx) + draft_model_input.b_mtp_index = model_input.b_mtp_index.index_select(0, select_idx) + draft_model_input.b_seq_len = model_input.b_seq_len.index_select(0, select_idx) + draft_model_input.mem_indexes = model_input.mem_indexes.index_select(0, select_idx) + draft_model_input.mem_indexes_cpu = None + draft_model_input.mtp_draft_input_hiddens = model_output.mtp_main_output_hiddens.index_select(0, select_idx) + draft_model_input.multimodal_params = [{"images": [], "audios": []} for _ in range(common_req_num)] + return draft_model_input + + def _pad_draft_input_to(self, draft_model_input: ModelInput, target_req_num: int): + """ + 将 shrink 到 real_req_num 行的 draft 输入再 padding 回 target_req_num (= common_req_num) 行, + 使本 rank 的 draft forward 行数与其它 dp rank 对齐,保证 MoE all-to-all / dp all-gather 的 + shape 一致。padding 行采用与 padded_prepare_decode_inputs 相同的占位约定: + b_req_idx -> HOLD_REQUEST_ID, mem_indexes -> HOLD_TOKEN_MEMINDEX。 + """ + cur_req_num = draft_model_input.batch_size + pad_num = target_req_num - cur_req_num + if pad_num <= 0: + return draft_model_input + + hold_req_id = g_infer_context.req_manager.HOLD_REQUEST_ID + hold_mem_idx = g_infer_context.req_manager.mem_manager.HOLD_TOKEN_MEMINDEX + + draft_model_input.input_ids = F.pad(draft_model_input.input_ids, (0, pad_num), value=0) + draft_model_input.b_req_idx = F.pad(draft_model_input.b_req_idx, (0, pad_num), value=hold_req_id) + draft_model_input.b_mtp_index = F.pad(draft_model_input.b_mtp_index, (0, pad_num), value=0) + # padding 行用一个合法的小 seq_len (沿用 padded_prepare_decode_inputs 中 fake req 的约定值 2) + draft_model_input.b_seq_len = F.pad(draft_model_input.b_seq_len, (0, pad_num), value=2) + draft_model_input.mem_indexes = F.pad(draft_model_input.mem_indexes, (0, pad_num), value=hold_mem_idx) + # mtp_draft_input_hiddens 为 (rows, hidden),沿 dim0 在尾部补 0 行 + draft_model_input.mtp_draft_input_hiddens = F.pad( + draft_model_input.mtp_draft_input_hiddens, (0, 0, 0, pad_num), value=0 + ) + draft_model_input.multimodal_params = draft_model_input.multimodal_params + [ + {"images": [], "audios": []} for _ in range(pad_num) + ] + draft_model_input.batch_size = target_req_num + draft_model_input.total_token_num = target_req_num * draft_model_input.max_kv_seq_len + return draft_model_input + def _draft_decode_eagle( self, model_input: ModelInput, @@ -582,19 +658,12 @@ def _draft_decode_eagle( mtp_accept_len: torch.Tensor, req_num: int, ): - all_next_token_ids = [] - # share some inference info with the main model - draft_model_input = model_input - draft_model_output = model_output - all_next_token_ids.append(next_token_ids) - draft_next_token_ids_gpu = torch.zeros((model_input.batch_size), dtype=torch.int64, device="cuda") - if req_num > 0: - draft_next_token_ids_gpu[:req_num].copy_(next_token_ids, non_blocking=True) - - real_req_num = req_num // (self.mtp_step + 1) - padded_req_num = model_input.batch_size // (self.mtp_step + 1) - real_req_num - eagle_mem_indexes_cpu = None + mtp_size = self.mtp_step + 1 + real_req_num = req_num // mtp_size + common_req_num = model_input.batch_size // mtp_size + padded_req_num = common_req_num - real_req_num + # 即使本 rank 没有真实请求, 也要为其它 rank 同步运行 mtp_step 次 draft forward 的集合通信。 g_infer_state_lock.acquire() if g_infer_context.radix_cache is not None: g_infer_context.radix_cache.free_radix_cache_to_get_enough_token(real_req_num * self.mtp_step) @@ -602,40 +671,58 @@ def _draft_decode_eagle( g_infer_state_lock.release() eagle_mem_indexes = eagle_mem_indexes_cpu.cuda(non_blocking=True) - # process the draft model output - for _step in range(self.mtp_step): + if real_req_num > 0: + ( + draft_model_input, + draft_next_token_ids, + accepted_req_idx, + ) = self._build_eagle_accepted_draft_input( + main_model_input=model_input, + main_model_output=model_output, + next_token_ids=next_token_ids, + mtp_accept_len=mtp_accept_len, + b_req_mtp_start_loc=b_req_mtp_start_loc, + ) + if padded_req_num > 0: + draft_model_input = self._pad_draft_input_to(draft_model_input, common_req_num) + draft_next_token_ids = F.pad(draft_next_token_ids, (0, padded_req_num), value=0) + + main_mem_indexes = model_input.mem_indexes.view(common_req_num, mtp_size) + eagle_padded = F.pad( + eagle_mem_indexes.view(self.mtp_step, real_req_num).transpose(0, 1).contiguous(), + (0, 0, 0, padded_req_num), + value=g_infer_context.req_manager.mem_manager.HOLD_TOKEN_MEMINDEX, + ) # (common_req_num, mtp_step) + mem_index_plan = torch.cat([main_mem_indexes, eagle_padded], dim=1) + accepted_offsets = F.pad(mtp_accept_len.long() - 1, (0, padded_req_num), value=0) + req_offsets = torch.arange(common_req_num, dtype=torch.long, device=mem_index_plan.device) + else: + # 本 rank 无真实请求: 纯 padding draft 输入, 仅用于跟随集合通信, 结果不写回。 + draft_model_input = self._build_padding_draft_input(model_input, model_output, common_req_num) + draft_next_token_ids = torch.zeros((common_req_num,), dtype=torch.int64, device="cuda") + mem_index_plan = model_input.mem_indexes.view(common_req_num, mtp_size) + accepted_offsets = torch.zeros((common_req_num,), dtype=torch.long, device=mem_index_plan.device) + req_offsets = torch.arange(common_req_num, dtype=torch.long, device=mem_index_plan.device) - draft_model_input.input_ids = draft_next_token_ids_gpu - draft_model_input.mtp_draft_input_hiddens = draft_model_output.mtp_main_output_hiddens - # spec decode: MTP + draft_model_output = model_output + all_next_token_ids = [draft_next_token_ids] + for _step in range(self.mtp_step): + draft_model_input.input_ids = draft_next_token_ids + if _step > 0: + draft_model_input.mtp_draft_input_hiddens = draft_model_output.mtp_main_output_hiddens + draft_model_input.mem_indexes = mem_index_plan[req_offsets, accepted_offsets + _step] draft_model_idx = _step % self.num_mtp_models draft_model_output: ModelOutput = self.draft_models[draft_model_idx].forward(draft_model_input) - # update the meta info of the inference + draft_next_token_ids = self._gen_argmax_token_ids(draft_model_output) draft_model_input.b_seq_len += 1 draft_model_input.max_kv_seq_len += 1 - eagle_mem_indexes_i = eagle_mem_indexes[_step * real_req_num : (_step + 1) * real_req_num] - eagle_mem_indexes_i = F.pad( - input=eagle_mem_indexes_i, - pad=(0, padded_req_num), - mode="constant", - value=g_infer_context.req_manager.mem_manager.HOLD_TOKEN_MEMINDEX, - ) - draft_model_input.mem_indexes = torch.cat( - [draft_model_input.mem_indexes.view(-1, self.mtp_step + 1)[:, 1:], eagle_mem_indexes_i.view(-1, 1)], - dim=1, - ).view(-1) - draft_next_token_ids_gpu = self._gen_argmax_token_ids(draft_model_output) - all_next_token_ids.append(draft_next_token_ids_gpu) + all_next_token_ids.append(draft_next_token_ids) - if req_num > 0: - all_next_token_ids = torch.stack(all_next_token_ids, dim=1) # [batch_size, mtp_step + 1] - all_next_token_ids = all_next_token_ids[0:req_num, :] - mtp_scatter_next_token_ids( - req_to_next_token_ids=self.model.req_manager.req_sampling_params_manager.req_to_next_token_ids, - b_req_mtp_start_loc=b_req_mtp_start_loc, + if real_req_num > 0: + all_next_token_ids = torch.stack(all_next_token_ids, dim=1)[:real_req_num, :] + self._scatter_accepted_next_token_ids( + accepted_req_idx=accepted_req_idx[:real_req_num], all_next_token_ids=all_next_token_ids, - b_req_idx=model_input.b_req_idx[:req_num], - mtp_accept_len=mtp_accept_len, ) return eagle_mem_indexes_cpu @@ -689,7 +776,6 @@ def prefill_overlap_mtp(self, event_pack: OverlapEventPack, prefill_reqs: List[I draft_model_output0, draft_model_output1 = model_output0, model_output1 for draft_model_idx in range(self.num_mtp_models): - draft_model_input0 = prepare_mtp_prefill_inputs( model_input=draft_model_input0, b_next_token_ids=draft_next_token_ids_gpu0, @@ -741,17 +827,36 @@ def decode_overlap_mtp(self, event_pack: OverlapEventPack, decode_reqs: List[Inf ( model_input0, run_reqs0, - _, + padded_req_num0, model_input1, run_reqs1, - _, + padded_req_num1, ) = padded_overlap_prepare_decode_inputs(decode_reqs) req_num0, req_num1 = len(run_reqs0), len(run_reqs1) all_next_token_ids = [] b_mtp_index_cpu0 = model_input0.b_mtp_index b_mtp_index_cpu1 = model_input1.b_mtp_index - with torch.cuda.stream(g_infer_context.get_overlap_stream()): + if self.mtp_step > 0: + # 标记两个 micro-batch 的 verify decode 布局,每个 req 一个 accept 数量 + # (padding 出来的 fake req 记为 1)。run_reqs* 内每个真实 req 占 mtp_step+1 行, + # 取每组首行即可得到逐 req 的列表。不设置会让主模型 verify forward 走非 verify 布局, + # 命中 hybrid 主模型从未捕获的 cudagraph key (bs, False) -> KeyError。 + mtp_size = self.mtp_step + 1 + accept_lens0 = [r.mtp_accept_len for r in run_reqs0[::mtp_size]] + [1] * padded_req_num0 + accept_lens1 = [r.mtp_accept_len for r in run_reqs1[::mtp_size]] + [1] * padded_req_num1 + model_input0.b_num_accepted_tokens = g_pin_mem_manager.gen_from_list( + key="b_num_accepted_tokens_0", + data=accept_lens0, + dtype=torch.int32, + ) + model_input1.b_num_accepted_tokens = g_pin_mem_manager.gen_from_list( + key="b_num_accepted_tokens_1", + data=accept_lens1, + dtype=torch.int32, + ) + + with torch.cuda.stream(g_infer_context.get_overlap_stream()): model_output0, model_output1 = self.model.microbatch_overlap_decode(model_input0, model_input1) logits0 = model_output0.logits logits1 = model_output1.logits @@ -820,6 +925,11 @@ def decode_overlap_mtp(self, event_pack: OverlapEventPack, decode_reqs: List[Inf if req_num0 + req_num1 > 0: event_pack.notify_post_handle_and_wait_pre_post_handle() verify_event.synchronize() + # 写回每个 req 的本步 accept 数量,供下一步 verify 经 b_num_accepted_tokens 传入 + # GDN/linear-att verify kernel(据此提交 conv/ssm 递归状态的正确偏移)。chunked 路径 + # 在 chunked_prefill/impl.py 同样写回;dp 缺失会让状态停留在 accept=1 -> 状态错乱、精度崩塌。 + for req, accept_len in zip(decode_reqs, mtp_accept_len_cpu): + req.mtp_accept_len = int(accept_len) verify_ok_reqs = [run_reqs[i] for i in range(len(run_reqs)) if accepted_index_cpu[i] == 1] update_packs = self._pre_post_handle(verify_ok_reqs, is_chuncked_mode=False) @@ -882,8 +992,13 @@ def _draft_decode_vanilla_overlap( ): all_next_token_ids = [] all_next_token_ids.append(next_token_ids) - # share some inference info with the main model - draft_model_input0, draft_model_input1 = model_input0, model_input1 + # share some inference info with the main model. copy.copy 后清空 b_num_accepted_tokens, + # 使 draft (MTP) forward 走普通 decode 布局 (bs, False);否则会沿用主模型设置的 verify 布局, + # 命中 MTP draft 模型从未捕获的 cudagraph key (bs, True) -> KeyError。 + draft_model_input0 = copy.copy(model_input0) + draft_model_input1 = copy.copy(model_input1) + draft_model_input0.b_num_accepted_tokens = None + draft_model_input1.b_num_accepted_tokens = None draft_model_output0, draft_model_output1 = model_output0, model_output1 draft_next_token_ids_gpu0 = torch.zeros((model_input0.batch_size), dtype=torch.int64, device="cuda") @@ -897,7 +1012,6 @@ def _draft_decode_vanilla_overlap( # process the draft model output for draft_model_idx in range(self.mtp_step): - draft_model_input0.input_ids = draft_next_token_ids_gpu0 draft_model_input0.mtp_draft_input_hiddens = draft_model_output0.mtp_main_output_hiddens draft_model_input1.input_ids = draft_next_token_ids_gpu1 @@ -940,8 +1054,13 @@ def _draft_decode_eagle_overlap( ): all_next_token_ids = [] all_next_token_ids.append(next_token_ids) - # share some inference info with the main model - draft_model_input0, draft_model_input1 = model_input0, model_input1 + # share some inference info with the main model. copy.copy 后清空 b_num_accepted_tokens, + # 使 draft (MTP) forward 走普通 decode 布局 (bs, False);否则会沿用主模型设置的 verify 布局, + # 命中 MTP draft 模型从未捕获的 cudagraph key (bs, True) -> KeyError。 + draft_model_input0 = copy.copy(model_input0) + draft_model_input1 = copy.copy(model_input1) + draft_model_input0.b_num_accepted_tokens = None + draft_model_input1.b_num_accepted_tokens = None draft_model_output0, draft_model_output1 = model_output0, model_output1 draft_next_token_ids_gpu0 = torch.zeros((model_input0.batch_size), dtype=torch.int64, device="cuda") @@ -968,7 +1087,6 @@ def _draft_decode_eagle_overlap( # process the draft model output for _step in range(self.mtp_step): - draft_model_input0.input_ids = draft_next_token_ids_gpu0 draft_model_input0.mtp_draft_input_hiddens = draft_model_output0.mtp_main_output_hiddens draft_model_input1.input_ids = draft_next_token_ids_gpu1 diff --git a/lightllm/server/router/model_infer/mode_backend/mtp_model_factory.py b/lightllm/server/router/model_infer/mode_backend/mtp_model_factory.py new file mode 100644 index 0000000000..1b4ade1ac0 --- /dev/null +++ b/lightllm/server/router/model_infer/mode_backend/mtp_model_factory.py @@ -0,0 +1,33 @@ +from lightllm.models.deepseek_mtp.model import Deepseek3MTPModel +from lightllm.models.qwen3_moe_mtp.model import Qwen3MOEMTPModel +from lightllm.models.mistral_mtp.model import MistralMTPModel +from lightllm.models.glm4_moe_lite_mtp.model import Glm4MoeLiteMTPModel + + +def create_mtp_draft_model(model_type: str, mtp_mode: str, mtp_model_kvargs: dict): + """Single source of truth for (model_type, mtp_mode) -> MTP draft model (#10). + Shared by base_backend and the static MTP benchmark.""" + if model_type == "deepseek_v3": + assert mtp_mode in ["vanilla_with_att", "eagle_with_att"] + return Deepseek3MTPModel(mtp_model_kvargs) + elif model_type == "qwen3_moe": + assert mtp_mode in ["vanilla_no_att", "eagle_no_att"] + return Qwen3MOEMTPModel(mtp_model_kvargs) + elif model_type == "mistral": + assert mtp_mode in ["vanilla_no_att", "eagle_no_att"] + return MistralMTPModel(mtp_model_kvargs) + elif model_type == "glm4_moe_lite": + assert mtp_mode in ["vanilla_with_att", "eagle_with_att"] + return Glm4MoeLiteMTPModel(mtp_model_kvargs) + elif model_type in ("qwen3_5", "qwen3_5_text"): + assert mtp_mode in ["vanilla_with_att", "eagle_with_att"] + from lightllm.models.qwen3_5_mtp.model import Qwen3_5MTPModel + + return Qwen3_5MTPModel(mtp_model_kvargs) + elif model_type in ("qwen3_5_moe", "qwen3_5_moe_text"): + assert mtp_mode in ["vanilla_with_att", "eagle_with_att"] + from lightllm.models.qwen3_5_moe_mtp.model import Qwen3_5MoeMTPModel + + return Qwen3_5MoeMTPModel(mtp_model_kvargs) + else: + raise ValueError(f"Unsupported MTP model type: {model_type}") diff --git a/lightllm/utils/envs_utils.py b/lightllm/utils/envs_utils.py index 2bdd4005fa..89b9ed6feb 100644 --- a/lightllm/utils/envs_utils.py +++ b/lightllm/utils/envs_utils.py @@ -226,16 +226,20 @@ def enable_huge_page(): return enable_env_vars("LIGHTLLM_HUGE_PAGE_ENABLE") +def _mtp_added_layer_num(mtp_mode, mtp_step: int) -> int: + # Single source of truth for the mtp_mode -> added KV/full-att layer count (#9). + if mtp_mode == "eagle_with_att": + return 1 + if mtp_mode == "vanilla_with_att": + return mtp_step + return 0 + + @lru_cache(maxsize=None) def get_added_mtp_kv_layer_num() -> int: # mtp 模式下需要在mem manger上扩展draft model使用的layer - added_mtp_layer_num = 0 - if get_env_start_args().mtp_mode == "eagle_with_att": - added_mtp_layer_num += 1 - elif get_env_start_args().mtp_mode == "vanilla_with_att": - added_mtp_layer_num += get_env_start_args().mtp_step - - return added_mtp_layer_num + args = get_env_start_args() + return _mtp_added_layer_num(args.mtp_mode, args.mtp_step) @lru_cache(maxsize=None) diff --git a/lightllm/utils/kv_cache_utils.py b/lightllm/utils/kv_cache_utils.py index 494908cb10..2a089a9bf2 100644 --- a/lightllm/utils/kv_cache_utils.py +++ b/lightllm/utils/kv_cache_utils.py @@ -120,8 +120,12 @@ def calcu_cpu_cache_meta() -> "CpuKVCacheMeta": if args.mtp_mode is not None: # TODO 可能会存在不同mtp模式的精度问题 - assert is_linear_att_mixed_model(args.model_dir) is False, "linear att mixed model does not support mtp mode" - cpu_cache_meta.layer_num += get_added_mtp_kv_layer_num() + if is_linear_att_mixed_model(args.model_dir): + # Linear mixed models use one packed byte page; MTP draft full-attn + # slots are accounted in LinearAttCacheConfig.get_cpu_cache_big_page_bytes(). + pass + else: + cpu_cache_meta.layer_num += get_added_mtp_kv_layer_num() cpu_cache_page_num = int( (args.cpu_cache_storage_size * 1024 * 1024 * 1024) / (cpu_cache_meta.calcu_one_page_size()) diff --git a/test/acc/cpu_cache_roundtrip_test.py b/test/acc/cpu_cache_roundtrip_test.py new file mode 100644 index 0000000000..28ce96f583 --- /dev/null +++ b/test/acc/cpu_cache_roundtrip_test.py @@ -0,0 +1,86 @@ +"""Force the CPU KV-cache offload->restore path and check correctness. + +GSM8K can't exercise the CPU cache (one shared hot prefix, sub-page tails). +This driver builds N distinct, page-aligned, long prompts that overflow the +GPU KV budget so their KV is offloaded to CPU, then re-requests them so they +are restored from CPU. With greedy decoding the round-2 (CPU-restored) output +MUST be token-identical to round-1 (freshly computed). For the MTP build it +also tracks accept-rate (mtp_avg_token_per_step) which would degrade if the +draft full-attn slots were not persisted/restored correctly. +""" +import argparse +import sys +import requests +from concurrent.futures import ThreadPoolExecutor + + +def make_prompts(n, words_per_prompt): + prompts = [] + for i in range(n): + # Distinct, deterministic filler so each prompt is its own radix branch + # and long enough to span several 256-token pages. + filler = " ".join(f"item{i}-{j}" for j in range(words_per_prompt)) + prompts.append( + f"You are given list number {i}. The list is: {filler}. " + f"Question: briefly summarize what list number {i} contains. Answer:" + ) + return prompts + + +def gen(url, prompt, max_tokens): + data = { + "inputs": prompt, + "parameters": { + "temperature": 0.0, + "max_new_tokens": max_tokens, + "stop_sequences": None, + "repetition_penalty": 1.0, + "top_p": 1.0, + "top_k": 1, + }, + } + r = requests.post(url, json=data, timeout=120) + assert r.status_code == 200, f"{r.status_code}: {r.text}" + return r.json()["generated_text"][0] + + +def run_round(url, prompts, max_tokens, parallel): + out = [None] * len(prompts) + with ThreadPoolExecutor(max_workers=parallel) as ex: + futs = {ex.submit(gen, url, p, max_tokens): k for k, p in enumerate(prompts)} + for f in futs: + k = futs[f] + out[k] = f.result() + return out + + +def main(): + ap = argparse.ArgumentParser() + ap.add_argument("--host", default="http://127.0.0.1") + ap.add_argument("--port", type=int, default=8088) + ap.add_argument("--num-prompts", type=int, default=24) + ap.add_argument("--words-per-prompt", type=int, default=400) + ap.add_argument("--max-tokens", type=int, default=32) + ap.add_argument("--parallel", type=int, default=8) + args = ap.parse_args() + + url = f"{args.host}:{args.port}/generate" + prompts = make_prompts(args.num_prompts, args.words_per_prompt) + + print(f"Round 1 (cold compute): {len(prompts)} distinct prompts", flush=True) + r1 = run_round(url, prompts, args.max_tokens, args.parallel) + print("Round 2 (CPU-restored):", flush=True) + r2 = run_round(url, prompts, args.max_tokens, args.parallel) + + mismatches = [i for i in range(len(prompts)) if r1[i] != r2[i]] + print(f"\n=== RESULT ===") + print(f"prompts: {len(prompts)} identical: {len(prompts) - len(mismatches)} mismatches: {len(mismatches)}") + if mismatches: + for i in mismatches[:5]: + print(f" [#{i}] R1={r1[i]!r}\n R2={r2[i]!r}") + sys.exit(1) + print("PASS: round-2 (CPU-restored) output is token-identical to round-1.") + + +if __name__ == "__main__": + main() diff --git a/test/benchmark/static_inference/model_infer.py b/test/benchmark/static_inference/model_infer.py index f2c900af09..b93c5fee55 100644 --- a/test/benchmark/static_inference/model_infer.py +++ b/test/benchmark/static_inference/model_infer.py @@ -36,7 +36,7 @@ def test_model_inference(args): "graph_max_len_in_batch": args.max_req_total_len, "graph_max_batch_size": args.graph_max_batch_size, "mem_fraction": args.mem_fraction, - "max_req_num": 2048, + "max_req_num": 512, "batch_max_tokens": 1024, "run_mode": "normal", "max_seq_length": args.max_req_total_len, diff --git a/test/benchmark/static_inference/model_infer_mtp.py b/test/benchmark/static_inference/model_infer_mtp.py index 72f06a919c..ff31133ae2 100644 --- a/test/benchmark/static_inference/model_infer_mtp.py +++ b/test/benchmark/static_inference/model_infer_mtp.py @@ -1,4 +1,5 @@ import os +import copy import torch import numpy as np from multiprocessing import Queue @@ -9,42 +10,60 @@ from lightllm.models import get_model from lightllm.common.basemodel.batch_objs import ModelInput, ModelOutput from lightllm.server.core.objs.start_args_type import StartArgs -from torch.profiler import profile, record_function, ProfilerActivity +from torch.profiler import profile, ProfilerActivity from lightllm.utils.log_utils import init_logger -from lightllm.models.deepseek_mtp.model import Deepseek3MTPModel -import torch.cuda as cuda logger = init_logger(__name__) def init_mtp_model(args: StartArgs, kvargs, main_model): - mtp_step = args.mtp_step draft_models = [] os.environ["DISABLE_CHECK_MAX_LEN_INFER"] = "1" - mtp_model_kvargs = kvargs - mtp_model_kvargs.update( - { - "weight_dir": args.mtp_draft_model_dir, + + if args.mtp_mode in ["vanilla_with_att", "vanilla_no_att"]: + num_mtp_modules = args.mtp_step + elif args.mtp_mode in ["eagle_with_att", "eagle_no_att"]: + num_mtp_modules = 1 + else: + assert False, f"error mtp mode {args.mtp_mode}" + + for i in range(num_mtp_modules): + mtp_model_cfg, _ = PretrainedConfig.get_config_dict(args.mtp_draft_model_dir[i]) + model_type = mtp_model_cfg.get("model_type", "") + mtp_model_kvargs = { + "weight_dir": args.mtp_draft_model_dir[i], "max_total_token_num": main_model.mem_manager.size, - "disable_chunked_prefill": True, - "mtp_mode": args.mtp_mode, + "load_way": kvargs["load_way"], + "max_req_num": kvargs.get("max_req_num", 1000), + "max_seq_length": kvargs.get("max_seq_length", 1024 * 5), + "is_token_healing": False, + "return_all_prompt_logics": False, + "disable_chunked_prefill": args.disable_chunked_prefill, + "data_type": kvargs.get("data_type", "float16"), + "graph_max_batch_size": kvargs.get("graph_max_batch_size", 16), + "graph_max_len_in_batch": kvargs.get("graph_max_len_in_batch", 8196), + "disable_cudagraph": kvargs.get("disable_cudagraph", False), + "mem_fraction": kvargs["mem_fraction"], + "batch_max_tokens": kvargs.get("batch_max_tokens", None), + "quant_type": kvargs.get("quant_type", None), + "quant_cfg": kvargs.get("quant_cfg", None), + "run_mode": "normal", + "llm_prefill_att_backend": kvargs.get("llm_prefill_att_backend", args.llm_prefill_att_backend), + "llm_decode_att_backend": kvargs.get("llm_decode_att_backend", args.llm_decode_att_backend), + "vit_att_backend": kvargs.get("vit_att_backend", args.vit_att_backend), + "llm_kv_type": kvargs.get("llm_kv_type", args.llm_kv_type), + "llm_kv_quant_group_size": kvargs.get("llm_kv_quant_group_size", args.llm_kv_quant_group_size), "main_model": main_model, + "mtp_previous_draft_models": draft_models.copy(), + "mtp_mode": args.mtp_mode, } - ) - for i in range(mtp_step): - mtp_model_cfg, _ = PretrainedConfig.get_config_dict(args.mtp_draft_model_dir) - mtp_model_kvargs.update( - { - "weight_dir": args.spec_model_dir, - "max_total_token_num": main_model.mem_manager.size, - "disable_chunked_prefill": True, - "mtp_mode": args.mtp_mode, - "main_model": main_model, - "mem_layer_start": main_model.config["num_hidden_layers"] + i * mtp_model_cfg["num_hidden_layers"], - } - ) - draft_models.append(Deepseek3MTPModel(mtp_model_kvargs)) + + from lightllm.server.router.model_infer.mode_backend.mtp_model_factory import create_mtp_draft_model + + draft_models.append(create_mtp_draft_model(model_type, args.mtp_mode, mtp_model_kvargs)) + + logger.info(f"loaded mtp model class {draft_models[i].__class__}") return draft_models @@ -68,13 +87,22 @@ def test_model_inference_mtp(args): "max_total_token_num": args.max_total_token_num, "graph_max_len_in_batch": args.max_req_total_len, "graph_max_batch_size": args.graph_max_batch_size, - "mem_faction": args.mem_fraction, - "max_req_num": 2000, + "mem_fraction": args.mem_fraction, + # Static bench runs explicit batch sizes (<= a few hundred). The hybrid Qwen3.5 + # GDN req-state cache is sized max_req_num * (mtp_step + 1) at ~34 MB/slot, so the + # old default of 2000 alloc'd ~140 GB and OOM'd under MTP. 512 covers any realistic + # static batch sweep while keeping the GDN cache small. + "max_req_num": 512, "batch_max_tokens": 2048, "run_mode": "normal", "max_seq_length": args.max_req_total_len, - "spec_algo": args.spec_algo, "disable_cudagraph": args.disable_cudagraph, + "quant_cfg": args.quant_cfg, + "llm_prefill_att_backend": args.llm_prefill_att_backend, + "llm_decode_att_backend": args.llm_decode_att_backend, + "vit_att_backend": args.vit_att_backend, + "llm_kv_type": args.llm_kv_type, + "llm_kv_quant_group_size": args.llm_kv_quant_group_size, } proc = multiprocessing.Process( target=tppart_model_infer, @@ -113,28 +141,36 @@ def run_forward_once(args, input_len, output_len, batch_size, main_model, draft_ test_data = np.vstack([np.random.randint(0, 50256, input_len) for _ in range(batch_size)]) test_data = test_data.reshape(-1) - test_data = torch.from_numpy(test_data).cuda() + test_data = torch.from_numpy(test_data) b_req_idx = torch.tensor( - [main_model.req_manager.alloc() for _ in range(batch_size)], dtype=torch.int32, device="cuda" + [main_model.req_manager.alloc() for _ in range(batch_size)], dtype=torch.int32, device="cpu" ) - b_seq_len = torch.zeros(batch_size, dtype=torch.int32, device="cuda") - b_ready_cache_len = torch.zeros(batch_size, dtype=torch.int32, device="cuda") + b_seq_len = torch.zeros(batch_size, dtype=torch.int32, device="cpu") + b_ready_cache_len = torch.zeros(batch_size, dtype=torch.int32, device="cpu") for i in range(batch_size): b_seq_len[i] = input_len total_token_num = input_len * batch_size - mem_indexes = main_model.req_manager.mem_manager.alloc(test_data.shape[0]).cuda() + mem_indexes = main_model.req_manager.mem_manager.alloc(test_data.shape[0]) + b_mtp_index = torch.zeros(batch_size, dtype=torch.int32) + b_prefill_start_loc = b_seq_len.cumsum(dim=0, dtype=torch.int32) - b_seq_len # Main model Prefill model_input = ModelInput( batch_size=batch_size, total_token_num=total_token_num, + max_q_seq_len=input_len, + max_kv_seq_len=input_len, + max_cache_len=0, input_ids=test_data, - mem_indexes=mem_indexes, + mem_indexes_cpu=mem_indexes, b_req_idx=b_req_idx, + b_mtp_index=b_mtp_index, b_seq_len=b_seq_len, is_prefill=True, b_ready_cache_len=b_ready_cache_len, + b_prefill_start_loc=b_prefill_start_loc, + prefix_total_token_num=0, multimodal_params=[{"images": [], "audios": []} for _ in range(batch_size)], ) @@ -167,8 +203,22 @@ def run_forward_once(args, input_len, output_len, batch_size, main_model, draft_ torch.cuda.synchronize() + # Speculative width = args.mtp_step in BOTH modes (mirrors base_backend: self.mtp_step = + # args.mtp_step). The number of draft MODEL INSTANCES differs: vanilla loads mtp_step + # instances (each forwarded once), eagle loads ONE instance forwarded mtp_step times + # (chunked_prefill/impl.py: draft_models[_step % num_instances]). The verify batch always + # expands to (mtp_step + 1) rows per request. + spec_width = args.mtp_step + num_instances = len(draft_models) + # The draft prefill above produced (1 + num_instances) columns; pad/truncate to + # (spec_width + 1) so the decode verify batch matches the server's expand width. Only the + # SHAPE matters for throughput here (argmax over random inputs); token values do not. + while len(draft_ids) < spec_width + 1: + draft_ids.append(draft_ids[-1]) + draft_ids = draft_ids[: spec_width + 1] decode_input_ids = np.stack(draft_ids, axis=-1).reshape(-1) - decode_input_ids = torch.from_numpy(decode_input_ids).cuda() + decode_input_ids = torch.from_numpy(decode_input_ids) + mtp_step = spec_width # build main decode input: nopad_b_seq_idx = [] @@ -177,67 +227,167 @@ def run_forward_once(args, input_len, output_len, batch_size, main_model, draft_ nopad_max_len_in_batch = 0 for i in range(batch_size): - nopad_b_seq_idx.append(b_req_idx[i]) + nopad_b_seq_idx.append(b_req_idx[i].item()) seq_len = b_seq_len[i].item() nopad_b_seq_len.append(seq_len + 1) nopad_total_token_num += seq_len + 1 - nopad_max_len_in_batch = max(nopad_max_len_in_batch, b_seq_len[i] + 1) + nopad_max_len_in_batch = max(nopad_max_len_in_batch, seq_len + 1) - for step in range(len(draft_models)): - nopad_b_seq_idx.append(b_req_idx[i]) + for step in range(mtp_step): + nopad_b_seq_idx.append(b_req_idx[i].item()) nopad_b_seq_len.append(seq_len + step + 2) nopad_total_token_num += seq_len + step + 2 nopad_max_len_in_batch = max(nopad_max_len_in_batch, seq_len + step + 2) - nopad_b_seq_idx = torch.tensor(nopad_b_seq_idx, dtype=torch.int32, device="cuda") - nopad_b_seq_len = torch.tensor(nopad_b_seq_len, dtype=torch.int32, device="cuda") - mem_indexes = main_model.req_manager.mem_manager.alloc(batch_size * (len(draft_models) + 1)).cuda() + nopad_b_seq_idx = torch.tensor(nopad_b_seq_idx, dtype=torch.int32, device="cpu") + nopad_b_seq_len = torch.tensor(nopad_b_seq_len, dtype=torch.int32, device="cpu") + b_mtp_index = torch.arange(mtp_step + 1, dtype=torch.int32).repeat(batch_size) + mem_indexes = main_model.req_manager.mem_manager.alloc(batch_size * (mtp_step + 1)) model_input = ModelInput( - batch_size=batch_size * (len(draft_models) + 1), + batch_size=batch_size * (mtp_step + 1), total_token_num=nopad_total_token_num, + max_q_seq_len=1, + max_kv_seq_len=nopad_max_len_in_batch, input_ids=decode_input_ids, - mem_indexes=mem_indexes, + mem_indexes_cpu=mem_indexes, b_req_idx=nopad_b_seq_idx, + b_mtp_index=b_mtp_index, b_seq_len=nopad_b_seq_len, is_prefill=False, - multimodal_params=[{"images": [], "audios": []} for _ in range(batch_size * (len(draft_models) + 1))], + multimodal_params=[{"images": [], "audios": []} for _ in range(batch_size * (mtp_step + 1))], ) + # MTP verify layout. The main decode is a VERIFY forward over the (mtp_step+1)-expanded + # batch. Setting b_num_accepted_tokens (one entry per real request) flips is_mtp_verify=True + # so the hybrid GDN main model runs the fused spec-decode verify kernel — the production path. + # Without it the main decode silently takes the plain _gdn_decode_kernel on the S+1-expanded + # batch (whose rows share req_idx), colliding on the single widened conv slot and mismeasuring + # cost. accept_len is fixed at 1 (steady-state low-acceptance); the verify-forward COST is + # ~constant in accept_len (it always processes mtp_step+1 rows), so this faithfully measures + # per-step decode cost. Vary accept_len in [1, mtp_step+1] to sweep the acceptance regime. + accept_len = 1 + is_eagle = args.mtp_mode.startswith("eagle") + model_input.b_num_accepted_tokens = torch.full((batch_size,), accept_len, dtype=torch.int32, device="cuda") + req_offsets = torch.arange(batch_size, dtype=torch.long, device="cuda") + accepted_row_idx = req_offsets * (mtp_step + 1) + (accept_len - 1) + if is_eagle: + # EAGLE draft scratch slots (n_real * mtp_step), mirroring _draft_decode_eagle. Allocated + # once and reused across steps (throughput bench overwrites draft KV; no correctness check). + eagle_mem_indexes = main_model.req_manager.mem_manager.alloc(batch_size * mtp_step).cuda() + + # Prize-sizing profiler (need-to-fix #22): env-gated, eagle-only, additive. Times the verify + # forward vs the S-step draft chain to decide whether collapsing the chain into a CUDA graph is + # worth it. host_bound_ratio ~1 (or per_step flat across bs) => host/launch-bound => graph wins. + _mtp_profile = os.environ.get("MTP_PROFILE", "0") == "1" + _prof = {"verify_ms": 0.0, "draft_ms": 0.0, "draft_host_ms": 0.0, "n": 0, "per_step_ms": [0.0] * mtp_step} + # Main decode - for i in range(0, output_len, len(draft_models) + 1): + for i in range(0, output_len, mtp_step + 1): torch.cuda.synchronize() step_start_time = time.time() - model_output = main_model.forward( - model_input, - ) - prob_out = torch.softmax(model_output.logits, dim=-1) - predict_ids = torch.argmax(prob_out, dim=1, keepdim=True) - # draft decode - model_input.input_ids = predict_ids.reshape(-1) - model_input.mtp_draft_input_hiddens = model_output.mtp_main_output_hiddens - - for draft_model_id in range(len(draft_models)): - draft_model = draft_models[draft_model_id] - model_output = draft_model.forward( - model_input, + # --- main VERIFY forward: mtp_step+1 rows/req through the fused GDN verify kernel --- + if _mtp_profile and not warmup: + _ev_v0 = torch.cuda.Event(enable_timing=True) + _ev_v1 = torch.cuda.Event(enable_timing=True) + _ev_v0.record() + model_output = main_model.forward(model_input) + if _mtp_profile and not warmup: + _ev_v1.record() + predict_ids = torch.argmax(model_output.logits, dim=1, keepdim=True) + + if is_eagle: + # EAGLE draft: shrink to the single accepted row per request (1 row/req), then run the + # draft model mtp_step times. The Qwen3.5 MTP draft is full-attention and takes the + # plain decode layout (b_num_accepted_tokens=None). Mirrors chunked_prefill + # _build_eagle_accepted_draft_input + _draft_decode_eagle so the measured draft cost is + # the real n_real-row cost, not the (mtp_step+1)x-inflated full-batch cost. + main_mem = model_input.mem_indexes.view(batch_size, mtp_step + 1) + eagle_mem_by_req = eagle_mem_indexes.view(mtp_step, batch_size).transpose(0, 1).contiguous() + mem_index_plan = torch.cat([main_mem, eagle_mem_by_req], dim=1) + + draft_model_input = copy.copy(model_input) + draft_model_input.batch_size = batch_size + draft_model_input.total_token_num = batch_size * model_input.max_kv_seq_len + draft_model_input.b_num_accepted_tokens = None + draft_model_input.mem_indexes_cpu = None + draft_model_input.b_req_idx = model_input.b_req_idx.index_select(0, accepted_row_idx) + draft_model_input.b_seq_len = model_input.b_seq_len.index_select(0, accepted_row_idx) + draft_model_input.b_mtp_index = model_input.b_mtp_index.index_select(0, accepted_row_idx) + draft_model_input.input_ids = predict_ids.reshape(-1).index_select(0, accepted_row_idx) + draft_model_input.mtp_draft_input_hiddens = model_output.mtp_main_output_hiddens.index_select( + 0, accepted_row_idx ) - prob_out = torch.softmax(model_output.logits, dim=-1) - predict_ids = torch.argmax(prob_out, dim=1, keepdim=True) - model_input.input_ids = predict_ids.reshape(-1) - model_input.mtp_draft_input_hiddens = model_output.mtp_main_output_hiddens + draft_model_input.multimodal_params = [{"images": [], "audios": []} for _ in range(batch_size)] + + if _mtp_profile and not warmup: + _step_evs = [] + _ev_d0 = torch.cuda.Event(enable_timing=True) + _ev_d1 = torch.cuda.Event(enable_timing=True) + _host_t0 = time.time() + _ev_d0.record() + for _step in range(mtp_step): + draft_model_input.mem_indexes = mem_index_plan[req_offsets, (accept_len - 1) + _step] + draft_model = draft_models[_step % num_instances] + if _mtp_profile and not warmup: + _es = torch.cuda.Event(enable_timing=True) + _ee = torch.cuda.Event(enable_timing=True) + _es.record() + draft_output = draft_model.forward(draft_model_input) + if _mtp_profile and not warmup: + _ee.record() + _step_evs.append((_es, _ee)) + draft_model_input.input_ids = torch.argmax(draft_output.logits, dim=1, keepdim=True).reshape(-1) + draft_model_input.mtp_draft_input_hiddens = draft_output.mtp_main_output_hiddens + draft_model_input.b_seq_len = draft_model_input.b_seq_len + 1 + draft_model_input.max_kv_seq_len += 1 + if _mtp_profile and not warmup: + _ev_d1.record() + _host_t1 = time.time() + else: + # VANILLA draft: full (mtp_step+1)-expanded batch, plain decode layout. Mirrors + # chunked_prefill _draft_decode_vanilla (b_num_accepted_tokens cleared on a copy so the + # MTP draft model does not inherit the main model's verify layout / cudagraph key). + draft_model_input = copy.copy(model_input) + draft_model_input.b_num_accepted_tokens = None + draft_model_input.input_ids = predict_ids.reshape(-1) + draft_model_input.mtp_draft_input_hiddens = model_output.mtp_main_output_hiddens + for _step in range(mtp_step): + draft_model = draft_models[_step % num_instances] + draft_output = draft_model.forward(draft_model_input) + draft_model_input.input_ids = torch.argmax(draft_output.logits, dim=1, keepdim=True).reshape(-1) + draft_model_input.mtp_draft_input_hiddens = draft_output.mtp_main_output_hiddens - # accept all draft ids by default. - model_input.input_ids = predict_ids.reshape(-1) - model_input.mtp_draft_input_hiddens = model_output.mtp_main_output_hiddens torch.cuda.synchronize() + if _mtp_profile and not warmup and is_eagle and i >= 3 * (mtp_step + 1): + # skip first 3 macro-steps (lazy cudagraph capture / cache warmup) + _prof["verify_ms"] += _ev_v0.elapsed_time(_ev_v1) + _prof["draft_ms"] += _ev_d0.elapsed_time(_ev_d1) + _prof["draft_host_ms"] += (_host_t1 - _host_t0) * 1000.0 + for _k, (_es, _ee) in enumerate(_step_evs): + _prof["per_step_ms"][_k] += _es.elapsed_time(_ee) + _prof["n"] += 1 if i % 100 == 0 or i == output_len - 1: step_end_time = time.time() if get_current_rank_in_dp() == 0 and not warmup: step_time = step_end_time - step_start_time print(i, " step cost time:", step_time * 1000) - print(f"Decode throughput: {batch_size * (len(draft_models) + 1) * args.dp / step_time} tokens/s") + # Peak (all-accepted) throughput: mtp_step+1 candidate tokens per req per step. + print(f"Decode throughput: {batch_size * (mtp_step + 1) * args.dp / step_time} tokens/s") + + if _mtp_profile and is_eagle and _prof["n"] > 0 and get_current_rank_in_dp() == 0 and not warmup: + n = _prof["n"] + ps = ", ".join(f"{v / n:.3f}" for v in _prof["per_step_ms"]) + print(f"[MTP_PROFILE] bs={batch_size} S={mtp_step} steps={n}") + print(f"[MTP_PROFILE] verify_gpu_ms = {_prof['verify_ms'] / n:.3f}") + print(f"[MTP_PROFILE] draft_chain_gpu_ms = {_prof['draft_ms'] / n:.3f}") + print(f"[MTP_PROFILE] draft_chain_host_ms = {_prof['draft_host_ms'] / n:.3f} (host-enqueue, no sync)") + print(f"[MTP_PROFILE] per_draft_step_gpu_ms = [{ps}]") + print( + f"[MTP_PROFILE] host_bound_ratio = " + f"{_prof['draft_host_ms'] / max(_prof['draft_ms'], 1e-9):.3f} (~1 => host-bound => graph wins)" + ) main_model.mem_manager.free_all() main_model.req_manager.free_all() diff --git a/test/benchmark/static_inference/test_model.py b/test/benchmark/static_inference/test_model.py index 5b3751bcc3..461f289780 100644 --- a/test/benchmark/static_inference/test_model.py +++ b/test/benchmark/static_inference/test_model.py @@ -16,7 +16,7 @@ def test_model_infer(self): args = get_env_start_args() if args.data_type is None: args.data_type = get_dtype(args.model_dir) - if args.mtp_mode == "deepseekv3": + if args.mtp_mode is not None: test_model_inference_mtp(args) else: test_model_inference(args) diff --git a/unit_tests/benchmark/test_static_mtp_bench_sets_accept_tokens.py b/unit_tests/benchmark/test_static_mtp_bench_sets_accept_tokens.py new file mode 100644 index 0000000000..889190cbbc --- /dev/null +++ b/unit_tests/benchmark/test_static_mtp_bench_sets_accept_tokens.py @@ -0,0 +1,21 @@ +import pathlib +import re + +BENCH = pathlib.Path(__file__).resolve().parents[2] / "test/benchmark/static_inference/model_infer_mtp.py" + + +def test_main_decode_sets_b_num_accepted_tokens(): + src = BENCH.read_text() + # Main verify decode must SET the field (flips is_mtp_verify -> fused GDN verify kernel). + assert re.search(r"model_input\.b_num_accepted_tokens\s*=\s*torch\.full", src), ( + "static MTP bench no longer sets b_num_accepted_tokens on the main decode ModelInput; " + "the verify path (production fused GDN kernel) is not being exercised (#5)." + ) + + +def test_draft_inputs_clear_b_num_accepted_tokens(): + src = BENCH.read_text() + # Draft forwards must CLEAR it so they take the plain decode layout, mirroring production. + assert len(re.findall(r"draft_model_input\.b_num_accepted_tokens\s*=\s*None", src)) >= 2, ( + "static MTP bench draft inputs must clear b_num_accepted_tokens (eagle + vanilla)." + ) diff --git a/unit_tests/common/basemodel/test_fp8_decode_verify_narrowed.py b/unit_tests/common/basemodel/test_fp8_decode_verify_narrowed.py new file mode 100644 index 0000000000..a671550d31 --- /dev/null +++ b/unit_tests/common/basemodel/test_fp8_decode_verify_narrowed.py @@ -0,0 +1,59 @@ +import types +import torch +import pytest + +import lightllm.common.basemodel.attention.fa3.fp8 as fp8_mod +from lightllm.common.basemodel.attention.fa3.fp8 import Fp8Fa3DecodeAttState + + +def _make_verify_state(n_real, mtp_size, head_num=2, head_dim=8): + """Build an Fp8Fa3DecodeAttState as init_state would leave it in MTP-verify mode, + bypassing init_state. b_att_seq_len/page_table are NARROW (n_real); infer_state.b_seq_len + is the FULL expanded tensor (n_real*mtp_size) that must NOT be used as cache_seqlens.""" + state = object.__new__(Fp8Fa3DecodeAttState) + batch = n_real * mtp_size + state.b_att_seq_len = torch.full((n_real,), 16, dtype=torch.int32) + state.page_table = torch.zeros((n_real, 16), dtype=torch.int32) + state.cu_seqlens_q = torch.arange(0, (n_real + 1) * mtp_size, mtp_size, dtype=torch.int32) + state.cu_seqlens_k = torch.zeros((n_real + 1,), dtype=torch.int32) + state.decode_max_q_seq_len = mtp_size + state.infer_state = types.SimpleNamespace( + b_seq_len=torch.full((batch,), 16, dtype=torch.int32), + batch_size=batch, + ) + # k/v descale sized per real request (att_batch_size), indexed by layer + state.k_descale = torch.ones((1, n_real, head_num)) + state.v_descale = torch.ones((1, n_real, head_num)) + state.backend = types.SimpleNamespace(_find_layer_index=lambda k, v, att_state: 0) + return state, batch + + +def test_fp8_decode_uses_narrowed_cache_seqlens_and_causal(monkeypatch): + n_real, mtp_size, head_num, head_dim = 3, 4, 2, 8 + state, batch = _make_verify_state(n_real, mtp_size, head_num, head_dim) + + captured = {} + + def fake_flash(**kwargs): + captured.update(kwargs) + q = kwargs["q"] + return torch.zeros((q.shape[0], q.shape[1], q.shape[2])) + + def fake_quant(x, use_per_token_if_dynamic=True): + return x, torch.ones((x.shape[0], 1)) + + monkeypatch.setattr(fp8_mod, "flash_attn_with_kvcache", fake_flash) + monkeypatch.setattr(fp8_mod, "scaled_fp8_quant", fake_quant) + + q = torch.randn((batch, head_num, head_dim)) + k = torch.randn((batch, head_num, head_dim)) + v = torch.randn((batch, head_num, head_dim)) + + state._fp8_decode_att(q=q, k=k, v=v) + + # The KV-side seqlens must be the NARROW per-real-request tensor, matching page_table rows. + assert captured["cache_seqlens"] is state.b_att_seq_len + assert captured["cache_seqlens"].shape[0] == n_real + assert captured["cache_seqlens"].shape[0] == captured["page_table"].shape[0] + # Verify decode must be causal, like the non-fp8 sibling. + assert captured["causal"] is True diff --git a/unit_tests/common/basemodel/test_mtp_decode_cuda_graph.py b/unit_tests/common/basemodel/test_mtp_decode_cuda_graph.py new file mode 100644 index 0000000000..7fa76f824e --- /dev/null +++ b/unit_tests/common/basemodel/test_mtp_decode_cuda_graph.py @@ -0,0 +1,393 @@ +from types import SimpleNamespace + +import torch + +from lightllm.common.basemodel.basemodel import TpPartBaseModel +from lightllm.common.basemodel.batch_objs import ModelInput + + +def test_mtp_decode_cuda_graph_warmup_uses_verify_layout(): + from lightllm.common.basemodel.cuda_graph import CudaGraph + + graph = CudaGraph.__new__(CudaGraph) + graph.mtp_step = 2 + graph.graph_max_len_in_batch = 128 + + class FakeMemManager: + HOLD_TOKEN_MEMINDEX = -1 + + def alloc(self, size): + return torch.arange(size, dtype=torch.int32) + + model = SimpleNamespace( + req_manager=SimpleNamespace(HOLD_REQUEST_ID=99), + mem_manager=FakeMemManager(), + _gen_special_model_input=lambda token_num: {"mtp_draft_input_hiddens": None}, + ) + + model_input = graph._build_warmup_decode_model_input(model, batch_size=6, device="cpu") + + assert model_input.batch_size == 6 + assert model_input.b_mtp_index.tolist() == [0, 1, 2, 0, 1, 2] + assert model_input.b_seq_len.tolist() == [2, 3, 4, 2, 3, 4] + assert model_input.b_num_accepted_tokens.tolist() == [1, 1] + assert model_input.total_token_num == 18 + + +def test_mtp_decode_cuda_graph_warmup_supports_normal_layout_for_draft(): + from lightllm.common.basemodel.cuda_graph import CudaGraph + + graph = CudaGraph.__new__(CudaGraph) + graph.mtp_step = 2 + graph.graph_max_len_in_batch = 128 + + class FakeMemManager: + HOLD_TOKEN_MEMINDEX = -1 + + def alloc(self, size): + return torch.arange(size, dtype=torch.int32) + + model = SimpleNamespace( + req_manager=SimpleNamespace(HOLD_REQUEST_ID=99), + mem_manager=FakeMemManager(), + _gen_special_model_input=lambda token_num: {"mtp_draft_input_hiddens": torch.full((token_num, 4), 3.0)}, + ) + + model_input = graph._build_warmup_decode_model_input( + model, + batch_size=5, + device="cpu", + is_mtp_verify_decode=False, + ) + + assert model_input.batch_size == 5 + assert model_input.b_mtp_index.tolist() == [0, 0, 0, 0, 0] + assert model_input.b_seq_len.tolist() == [2, 2, 2, 2, 2] + assert model_input.b_num_accepted_tokens is None + assert model_input.total_token_num == 10 + assert model_input.mtp_draft_input_hiddens.shape == (5, 4) + + +def test_mtp_decode_cuda_graph_keys_verify_and_normal_layouts(): + from lightllm.common.basemodel.cuda_graph import CudaGraph + + graph = CudaGraph.__new__(CudaGraph) + graph.mtp_step = 2 + graph.graph = {} + graph.normal_cuda_graph_batch_sizes = [1, 2, 4, 8] + graph.mtp_verify_cuda_graph_batch_sizes = [3, 6, 9, 12] + + verify_state = SimpleNamespace( + input_ids=torch.ones(6, dtype=torch.int64), + b_num_accepted_tokens=torch.ones(2, dtype=torch.int32), + ) + normal_state = SimpleNamespace( + input_ids=torch.ones(6, dtype=torch.int64), + b_num_accepted_tokens=None, + ) + + assert graph._decode_graph_key(verify_state) == (6, True) + assert graph._decode_graph_key(normal_state) == (6, False) + assert graph.find_closest_graph_batch_size(5, is_mtp_verify_decode=True) == 6 + assert graph.find_closest_graph_batch_size(5, is_mtp_verify_decode=False) == 8 + + graph.graph[(6, True)] = "verify graph" + assert graph.need_capture(6, is_mtp_verify_decode=True) is False + assert graph.need_capture(5, is_mtp_verify_decode=False) is True + + +def test_mtp_decode_cuda_graph_warmup_layouts_split_main_and_draft_models(): + from lightllm.common.basemodel.cuda_graph import CudaGraph + + class Qwen3_5MOETpPartModel: + pass + + class Qwen3_5MoeMTPModel: + pass + + graph = CudaGraph.__new__(CudaGraph) + graph.mtp_step = 2 + graph.normal_cuda_graph_batch_sizes = [1, 2, 4, 8] + graph.mtp_verify_cuda_graph_batch_sizes = [3, 6, 9] + + assert list(graph._iter_warmup_graph_layouts(Qwen3_5MOETpPartModel())) == [(True, [3, 6, 9])] + assert list(graph._iter_warmup_graph_layouts(Qwen3_5MoeMTPModel())) == [(False, [1, 2, 4, 8])] + + +def test_mtp_decode_warmup_layout_marks_qwen3next_verify(monkeypatch): + import pytest + + if not torch.cuda.is_available(): + pytest.skip("needs CUDA for gen_decode_params") + + import lightllm.models.qwen3next.infer_struct as infer_struct_mod + from lightllm.models.qwen3next.infer_struct import Qwen3NextInferStateInfo + + monkeypatch.setattr(infer_struct_mod, "get_env_start_args", lambda: SimpleNamespace(mtp_step=2)) + + state = Qwen3NextInferStateInfo() + state.is_prefill = False + state.b_req_idx = torch.tensor([5, 5, 5, 6, 6, 6], dtype=torch.int32, device="cuda") + state.b_mtp_index = torch.tensor([0, 1, 2, 0, 1, 2], dtype=torch.int32, device="cuda") + state.b_seq_len = torch.tensor([2, 3, 4, 2, 3, 4], dtype=torch.int32, device="cuda") + state.b_num_accepted_tokens = torch.tensor([1, 2], dtype=torch.int32, device="cuda") + + model = SimpleNamespace( + _cos_cached=torch.zeros((16, 4), dtype=torch.float32, device="cuda"), + _sin_cached=torch.zeros((16, 4), dtype=torch.float32, device="cuda"), + ) + + state.init_some_extra_state(model) + + assert state.is_mtp_verify is True + assert state.b_gdn_verify_cu_seqlens.tolist() == [0, 3, 6] + assert state.b_conv_buffer_idx.tolist() == [5, 6] + assert state.b_ssm_index_rows.tolist() == [[15, 16, 17], [18, 19, 20]] + + +def test_mtp_decode_padding_preserves_verify_groups(monkeypatch): + import lightllm.common.basemodel.basemodel as basemodel_mod + + monkeypatch.setattr(basemodel_mod, "enable_diverse_mode_gqa_decode_fast_kernel", lambda: False) + + model = TpPartBaseModel.__new__(TpPartBaseModel) + model.args = SimpleNamespace(mtp_step=2) + model.req_manager = SimpleNamespace(HOLD_REQUEST_ID=99) + model.mem_manager = SimpleNamespace(HOLD_TOKEN_MEMINDEX=-1) + + model_input = ModelInput( + batch_size=3, + total_token_num=12, + max_q_seq_len=1, + max_kv_seq_len=4, + input_ids=torch.tensor([10, 11, 12], dtype=torch.int32), + mem_indexes=torch.tensor([20, 21, 22], dtype=torch.int32), + b_req_idx=torch.tensor([7, 7, 7], dtype=torch.int32), + b_mtp_index=torch.tensor([0, 1, 2], dtype=torch.int32), + b_seq_len=torch.tensor([2, 3, 4], dtype=torch.int32), + b_num_accepted_tokens=torch.tensor([2], dtype=torch.int32), + is_prefill=False, + multimodal_params=[{"images": [], "audios": []} for _ in range(3)], + ) + + padded = model._create_padded_decode_model_input(model_input, new_batch_size=6) + + assert padded.batch_size == 6 + assert padded.b_req_idx.tolist() == [7, 7, 7, 99, 99, 99] + assert padded.b_mtp_index.tolist() == [0, 1, 2, 0, 1, 2] + assert padded.b_seq_len.tolist() == [2, 3, 4, 2, 3, 4] + assert padded.b_num_accepted_tokens.tolist() == [2, 1] + assert padded.mem_indexes.tolist() == [20, 21, 22, -1, -1, -1] + assert len(padded.multimodal_params) == 6 + + +def test_qwen3next_hybrid_mtp_keeps_decode_cuda_graph_enabled(monkeypatch): + import lightllm.models.qwen3next.model as qwen3next_model + from lightllm.models.qwen3next.model import Qwen3NextTpPartModel + + monkeypatch.setattr(qwen3next_model, "get_env_start_args", lambda: SimpleNamespace(mtp_step=2)) + + called = {} + + def fake_base_init_cudagraph(self): + called["disable_cudagraph"] = self.disable_cudagraph + self.graph = "captured" + + monkeypatch.setattr(TpPartBaseModel, "_init_cudagraph", fake_base_init_cudagraph) + + model = Qwen3NextTpPartModel.__new__(Qwen3NextTpPartModel) + model.disable_cudagraph = False + + Qwen3NextTpPartModel._init_cudagraph(model) + + assert called["disable_cudagraph"] is False + assert model.disable_cudagraph is False + assert model.graph == "captured" + + +def test_fa3_decode_uses_normal_layout_for_narrowed_mtp_draft(monkeypatch): + import lightllm.common.basemodel.attention.fa3.fp as fa3_fp + from lightllm.common.basemodel.attention.fa3.fp import Fa3DecodeAttState + + monkeypatch.setattr(fa3_fp, "get_env_start_args", lambda: SimpleNamespace(mtp_step=2)) + + copied = {} + + def fake_page_table_copy(page_table, req_to_token_indexs, b_req_idx): + copied["page_table_shape"] = tuple(page_table.shape) + copied["b_req_idx"] = b_req_idx.clone() + + monkeypatch.setattr(fa3_fp, "page_table_copy", fake_page_table_copy) + + model = SimpleNamespace( + graph_max_batch_size=16, + graph_max_len_in_batch=32, + req_manager=SimpleNamespace(req_to_token_indexs=torch.empty((8, 32), dtype=torch.int32)), + ) + backend = SimpleNamespace( + model=model, + get_page_table_buffer=lambda: [torch.empty(16 * 32, dtype=torch.int32)], + ) + infer_state = SimpleNamespace( + batch_size=2, + max_kv_seq_len=16, + input_ids=torch.ones(2, dtype=torch.int64), + b_seq_len=torch.tensor([5, 7], dtype=torch.int32), + b1_cu_q_seq_len=torch.tensor([0, 1, 2], dtype=torch.int32), + b1_cu_kv_seq_len=torch.tensor([0, 5, 12], dtype=torch.int32), + b_req_idx=torch.tensor([3, 4], dtype=torch.int32), + b_num_accepted_tokens=None, + microbatch_index=0, + ) + + state = Fa3DecodeAttState(backend=backend, infer_state=infer_state) + state.init_state() + + assert state.decode_max_q_seq_len == 1 + assert state.b_att_seq_len.tolist() == [5, 7] + assert copied["page_table_shape"] == (2, 16) + assert copied["b_req_idx"].tolist() == [3, 4] + + +def test_build_eagle_accepted_draft_input_narrows_to_accepted_rows(): + from lightllm.common.basemodel.batch_objs import ModelInput, ModelOutput + from lightllm.server.router.model_infer.mode_backend.chunked_prefill.impl import ( + ChunkedPrefillBackend, + ) + + backend = ChunkedPrefillBackend.__new__(ChunkedPrefillBackend) + backend.mtp_step = 2 + + main_input = ModelInput( + batch_size=6, + total_token_num=27, + max_q_seq_len=1, + max_kv_seq_len=9, + input_ids=torch.tensor([10, 11, 12, 20, 21, 22], dtype=torch.int64), + mem_indexes=torch.tensor([100, 101, 102, 200, 201, 202], dtype=torch.int32), + b_req_idx=torch.tensor([3, 3, 3, 4, 4, 4], dtype=torch.int32), + b_mtp_index=torch.tensor([0, 1, 2, 0, 1, 2], dtype=torch.int32), + b_seq_len=torch.tensor([5, 6, 7, 6, 7, 8], dtype=torch.int32), + b_num_accepted_tokens=torch.tensor([1, 1], dtype=torch.int32), + is_prefill=False, + multimodal_params=[ + {"row": 0}, + {"row": 1}, + {"row": 2}, + {"row": 3}, + {"row": 4}, + {"row": 5}, + ], + ) + hidden = torch.arange(6 * 4, dtype=torch.float32).view(6, 4) + main_output = ModelOutput(logits=torch.empty(6, 8), mtp_main_output_hiddens=hidden) + next_token_ids = torch.tensor([110, 111, 112, 220, 221, 222], dtype=torch.int64) + b_req_mtp_start_loc = torch.tensor([0, 3], dtype=torch.int32) + mtp_accept_len = torch.tensor([2, 3], dtype=torch.int32) + + (draft_input, accepted_next_tokens, accepted_req_idx,) = backend._build_eagle_accepted_draft_input( + main_model_input=main_input, + main_model_output=main_output, + next_token_ids=next_token_ids, + mtp_accept_len=mtp_accept_len, + b_req_mtp_start_loc=b_req_mtp_start_loc, + ) + + assert draft_input.batch_size == 2 + assert draft_input.input_ids.tolist() == [111, 222] + assert draft_input.b_req_idx.tolist() == [3, 4] + assert draft_input.b_mtp_index.tolist() == [1, 2] + assert draft_input.b_seq_len.tolist() == [6, 8] + assert draft_input.mem_indexes.tolist() == [101, 202] + assert draft_input.b_num_accepted_tokens is None + assert draft_input.multimodal_params == [{"row": 1}, {"row": 5}] + assert accepted_next_tokens.tolist() == [111, 222] + assert accepted_req_idx.tolist() == [3, 4] + torch.testing.assert_close(draft_input.mtp_draft_input_hiddens, hidden[[1, 5]]) + + +def test_eagle_draft_decode_uses_narrowed_hidden_on_first_step(monkeypatch): + import lightllm.server.router.model_infer.mode_backend.chunked_prefill.impl as chunked_impl + from lightllm.common.basemodel.batch_objs import ModelInput, ModelOutput + from lightllm.server.router.model_infer.mode_backend.chunked_prefill.impl import ( + ChunkedPrefillBackend, + ) + + class FakeMemManager: + HOLD_TOKEN_MEMINDEX = -1 + + def alloc(self, need_size): + return torch.arange(300, 300 + need_size, dtype=torch.int32) + + req_to_next_token_ids = torch.empty((8, 3), dtype=torch.int64) + monkeypatch.setattr( + chunked_impl, + "g_infer_context", + SimpleNamespace( + radix_cache=None, + req_manager=SimpleNamespace( + mem_manager=FakeMemManager(), + req_sampling_params_manager=SimpleNamespace(req_to_next_token_ids=req_to_next_token_ids), + ), + ), + ) + monkeypatch.setattr(torch.Tensor, "cuda", lambda self, non_blocking=False: self) + + backend = ChunkedPrefillBackend.__new__(ChunkedPrefillBackend) + backend.mtp_step = 2 + backend.num_mtp_models = 1 + + seen_hiddens = [] + + class FakeDraftModel: + def forward(self, model_input): + seen_hiddens.append(model_input.mtp_draft_input_hiddens.clone()) + logits = torch.zeros((model_input.batch_size, 8), dtype=torch.float32) + return ModelOutput( + logits=logits, + mtp_main_output_hiddens=model_input.mtp_draft_input_hiddens + 100, + ) + + backend.draft_models = [FakeDraftModel()] + + scattered = {} + + def fake_scatter(accepted_req_idx, all_next_token_ids): + scattered["accepted_req_idx"] = accepted_req_idx.clone() + scattered["all_next_token_ids"] = all_next_token_ids.clone() + + backend._scatter_accepted_next_token_ids = fake_scatter + + main_input = ModelInput( + batch_size=6, + total_token_num=27, + max_q_seq_len=1, + max_kv_seq_len=9, + input_ids=torch.tensor([10, 11, 12, 20, 21, 22], dtype=torch.int64), + mem_indexes=torch.tensor([100, 101, 102, 200, 201, 202], dtype=torch.int32), + b_req_idx=torch.tensor([3, 3, 3, 4, 4, 4], dtype=torch.int32), + b_mtp_index=torch.tensor([0, 1, 2, 0, 1, 2], dtype=torch.int32), + b_seq_len=torch.tensor([5, 6, 7, 6, 7, 8], dtype=torch.int32), + b_num_accepted_tokens=torch.tensor([1, 1], dtype=torch.int32), + is_prefill=False, + multimodal_params=[{"images": [], "audios": []} for _ in range(6)], + ) + hidden = torch.arange(6 * 4, dtype=torch.float32).view(6, 4) + main_output = ModelOutput(logits=torch.empty(6, 8), mtp_main_output_hiddens=hidden) + next_token_ids = torch.tensor([110, 111, 112, 220, 221, 222], dtype=torch.int64) + b_req_mtp_start_loc = torch.tensor([0, 3], dtype=torch.int32) + mtp_accept_len = torch.tensor([2, 3], dtype=torch.int32) + + returned_mem = backend._draft_decode_eagle( + main_model_input=main_input, + main_model_output=main_output, + next_token_ids=next_token_ids, + mtp_accept_len=mtp_accept_len, + b_req_mtp_start_loc=b_req_mtp_start_loc, + ) + + assert returned_mem.tolist() == [300, 301, 302, 303] + torch.testing.assert_close(seen_hiddens[0], hidden[[1, 5]]) + torch.testing.assert_close(seen_hiddens[1], hidden[[1, 5]] + 100) + assert scattered["accepted_req_idx"].tolist() == [3, 4] + assert scattered["all_next_token_ids"].shape == (2, 3) diff --git a/unit_tests/common/test_conv_state_shape_split.py b/unit_tests/common/test_conv_state_shape_split.py new file mode 100644 index 0000000000..accb1095ce --- /dev/null +++ b/unit_tests/common/test_conv_state_shape_split.py @@ -0,0 +1,33 @@ +import torch +import pytest + + +def _make_cfg(conv_kernel_size=4, mtp_step=0): + from lightllm.common.linear_att_cache_manager.config_objs import LinearAttCacheConfig + + return LinearAttCacheConfig( + tp_world_size=1, + full_att_all_num_kv_heads=16, + full_att_dtype=torch.bfloat16, + full_att_num_kv_heads=16, + full_att_head_dim=256, + num_linear_k_heads=16, + num_linear_v_heads=48, + head_linear_k_dim=128, + head_linear_v_dim=128, + conv_kernel_size=conv_kernel_size, + linear_layer_num=48, + conv_state_dtype=torch.bfloat16, + ssm_state_dtype=torch.bfloat16, + full_attention_interval=4, + all_layer_num=64, + ) + + +@pytest.mark.parametrize("S", [0, 1, 2, 3]) +def test_gpu_shape_widens_by_S_persisted_stays_narrow(S): + cfg = _make_cfg(conv_kernel_size=4, mtp_step=S) + conv_dim = cfg.get_conv_dim() + assert cfg.get_persisted_conv_state_shape() == (conv_dim, 4 - 1) + assert cfg.get_gpu_conv_state_shape(mtp_step=S) == (conv_dim, (4 - 1) + S) + assert cfg.get_conv_state_bytes_per_layer() == conv_dim * (4 - 1) * cfg.conv_state_dtype.itemsize diff --git a/unit_tests/common/test_init_linear_att_state_zeros_block.py b/unit_tests/common/test_init_linear_att_state_zeros_block.py new file mode 100644 index 0000000000..b20e489bce --- /dev/null +++ b/unit_tests/common/test_init_linear_att_state_zeros_block.py @@ -0,0 +1,41 @@ +import types +import torch + +# NOTE: importing lightllm.common.req_manager *first* trips a pre-existing circular import +# (req_manager line-8 imports gen_sampling_params -> basemodel -> infer_struct, which re-enters +# the half-initialized req_manager before ReqManager is defined). Importing basemodel first +# fully resolves that chain, after which ReqManagerForMamba imports cleanly. This is an +# import-ordering fix only; it does not alter the method-under-test or the duck-typed call below. +import lightllm.common.basemodel # noqa: F401 (resolves circular import; must precede req_manager) +from lightllm.common.req_manager import ReqManagerForMamba + + +class _Buf: + def __init__(self, t): + self.buffer = t + + +def test_init_zeros_full_ssm_block(): + mtp_step = 3 + layer, n_req = 2, 4 + conv_dim, width = 8, 3 + conv_buf = torch.ones(layer, n_req, conv_dim, width) + ssm_buf = torch.ones(layer, n_req * (mtp_step + 1), 5) + + dummy = types.SimpleNamespace( + mtp_step=mtp_step, + req_to_conv_state=_Buf(conv_buf), + req_to_ssm_state=_Buf(ssm_buf), + ) + req = types.SimpleNamespace(req_idx=2, mtp_accept_len=None) + + ReqManagerForMamba.init_linear_att_state(dummy, req) + + start = 2 * (mtp_step + 1) + block = ssm_buf[:, start : start + (mtp_step + 1), ...] + assert torch.count_nonzero(block) == 0, "all S+1 SSM rows of the block must be zeroed on init" + # other requests' rows must be untouched + assert torch.count_nonzero(ssm_buf[:, :start, ...]) > 0 + # conv slot for this request zeroed; canonical accept-len reset + assert torch.count_nonzero(conv_buf[:, 2, ...]) == 0 + assert req.mtp_accept_len == 1 diff --git a/unit_tests/common/test_is_mtp_draft_model_detection.py b/unit_tests/common/test_is_mtp_draft_model_detection.py new file mode 100644 index 0000000000..082d35a398 --- /dev/null +++ b/unit_tests/common/test_is_mtp_draft_model_detection.py @@ -0,0 +1,20 @@ +import types + + +def test_detection_uses_attribute_not_string(): + from lightllm.common.basemodel.cuda_graph import CudaGraph + from lightllm.models.base_mtp_model import BaseMTPModel + + graph = CudaGraph.__new__(CudaGraph) + + class _Draft(BaseMTPModel): + pass + + class _Main: + pass + + assert graph._is_mtp_draft_model(_Draft.__new__(_Draft)) is True + assert graph._is_mtp_draft_model(_Main()) is False + # a non-MTP class whose name happens to contain "MTPModel" must NOT be misdetected + GotchaMTPModelButNot = types.new_class("GotchaMTPModelButNot") + assert graph._is_mtp_draft_model(GotchaMTPModelButNot()) is False diff --git a/unit_tests/common/test_is_mtp_verify_decode_predicate.py b/unit_tests/common/test_is_mtp_verify_decode_predicate.py new file mode 100644 index 0000000000..313ea7a351 --- /dev/null +++ b/unit_tests/common/test_is_mtp_verify_decode_predicate.py @@ -0,0 +1,8 @@ +def test_predicate(): + from lightllm.common.basemodel.batch_objs import is_mtp_verify_decode + + sentinel = object() + assert is_mtp_verify_decode(3, sentinel) is True + assert is_mtp_verify_decode(3, None) is False + assert is_mtp_verify_decode(0, sentinel) is False + assert is_mtp_verify_decode(0, None) is False diff --git a/unit_tests/common/test_linear_att_copy_guards.py b/unit_tests/common/test_linear_att_copy_guards.py new file mode 100644 index 0000000000..b6c48a0c86 --- /dev/null +++ b/unit_tests/common/test_linear_att_copy_guards.py @@ -0,0 +1,39 @@ +import pytest +import torch + +from lightllm.common.basemodel.triton_kernel.linear_att_copy import ( + copy_linear_att_state_to_kv_buffer, +) + + +def _args(gpu_conv, accept_len, mtp_step): + layer_num = gpu_conv.shape[0] + dim_conv = gpu_conv.shape[2] + width_narrow = 3 + return dict( + b_req_idx=torch.tensor([0], dtype=torch.int32), + big_page_buffer_ids=torch.tensor([0], dtype=torch.int32), + gpu_conv_state=gpu_conv, + gpu_ssm_state=torch.zeros(layer_num, 1 * (mtp_step + 1), 8), + cpu_kv_conv_state=torch.zeros(1, layer_num, dim_conv, width_narrow), + cpu_kv_ssm_state=torch.zeros(1, layer_num, 8), + mtp_step=mtp_step, + b_num_accepted_tokens=torch.tensor([accept_len], dtype=torch.int32), + ) + + +def test_rejects_non_contiguous_width_axis(): + mtp_step = 2 + # widened slot allocated 2x, then strided ::2 along the width axis -> stride(3) == 2 + base = torch.zeros(2, 1, 32, (3 + mtp_step) * 2) + gpu_conv = base[:, :, :, ::2] + assert gpu_conv.stride(3) != 1 + with pytest.raises(AssertionError, match="width"): + copy_linear_att_state_to_kv_buffer(**_args(gpu_conv, accept_len=1, mtp_step=mtp_step)) + + +def test_rejects_out_of_range_accept_len(): + mtp_step = 2 + gpu_conv = torch.zeros(2, 1, 32, 3 + mtp_step) # contiguous, passes the #6 guard + with pytest.raises(AssertionError, match="b_num_accepted_tokens"): + copy_linear_att_state_to_kv_buffer(**_args(gpu_conv, accept_len=mtp_step + 2, mtp_step=mtp_step)) diff --git a/unit_tests/common/test_linear_att_mtp_cpu_cache_persistence.py b/unit_tests/common/test_linear_att_mtp_cpu_cache_persistence.py new file mode 100644 index 0000000000..fb22f0ed1f --- /dev/null +++ b/unit_tests/common/test_linear_att_mtp_cpu_cache_persistence.py @@ -0,0 +1,219 @@ +from types import SimpleNamespace + +import pytest +import torch + + +def _make_start_args(**overrides): + base = dict( + model_dir="/tmp/qwen3_5", + tp=1, + dp=1, + data_type="bfloat16", + linear_att_ssm_data_type="bfloat16", + mtp_mode=None, + mtp_step=0, + linear_att_page_block_num=2, + linear_att_hash_page_size=4, + cpu_cache_token_page_size=8, + ) + base.update(overrides) + return SimpleNamespace(**base) + + +def _make_model_cfg(): + return { + "model_type": "qwen3_5", + "num_hidden_layers": 64, + "num_key_value_heads": 16, + "head_dim": 128, + "linear_num_key_heads": 16, + "linear_num_value_heads": 48, + "linear_key_head_dim": 128, + "linear_value_head_dim": 128, + "linear_conv_kernel_dim": 4, + "full_attention_interval": 4, + } + + +def _patch_linear_config_args(monkeypatch, args): + import lightllm.common.linear_att_cache_manager.config_objs as config_objs + + monkeypatch.setattr(config_objs, "get_env_start_args", lambda: args) + + +def _make_config(draft_full_att_layer_num=0): + from lightllm.common.linear_att_cache_manager.config_objs import LinearAttCacheConfig + + return LinearAttCacheConfig( + tp_world_size=1, + full_att_all_num_kv_heads=16, + full_att_dtype=torch.bfloat16, + full_att_num_kv_heads=16, + full_att_head_dim=128, + num_linear_k_heads=16, + num_linear_v_heads=48, + head_linear_k_dim=128, + head_linear_v_dim=128, + conv_kernel_size=4, + linear_layer_num=48, + conv_state_dtype=torch.bfloat16, + ssm_state_dtype=torch.bfloat16, + full_attention_interval=4, + all_layer_num=64, + draft_full_att_layer_num=draft_full_att_layer_num, + ) + + +def test_load_from_args_includes_mtp_draft_full_att_layers(monkeypatch): + from lightllm.common.linear_att_cache_manager.config_objs import LinearAttCacheConfig + from transformers.configuration_utils import PretrainedConfig + + args = _make_start_args(mtp_mode="vanilla_with_att", mtp_step=3) + _patch_linear_config_args(monkeypatch, args) + monkeypatch.setattr(PretrainedConfig, "get_config_dict", lambda _model_path: (_make_model_cfg(), None)) + + cfg = LinearAttCacheConfig.load_from_args() + + assert cfg.get_main_full_att_layer_num() == 16 + assert cfg.draft_full_att_layer_num == 3 + assert cfg.get_persisted_full_att_layer_num() == 19 + + +def test_cpu_cache_full_att_bytes_include_mtp_draft_layers(monkeypatch): + args = _make_start_args() + _patch_linear_config_args(monkeypatch, args) + main_only = _make_config(draft_full_att_layer_num=0) + with_draft = _make_config(draft_full_att_layer_num=2) + + bytes_per_full_att_layer = ( + args.cpu_cache_token_page_size + * 2 + * main_only.full_att_all_num_kv_heads + * main_only.full_att_head_dim + * main_only.full_att_dtype.itemsize + ) + + assert main_only.get_main_full_att_layer_num() == 16 + assert with_draft.get_persisted_full_att_layer_num() == 18 + assert with_draft.get_cpu_cache_full_att_bytes() == ( + main_only.get_cpu_cache_full_att_bytes() + 2 * bytes_per_full_att_layer + ) + + +def test_linear_operator_persisted_full_att_slice_includes_draft_slots(): + from lightllm.common.kv_cache_mem_manager.operator.linear_att import LinearAttMemOperator + + class MtpMemManager: + main_full_att_layer_num = 16 + draft_full_att_layers = 2 + kv_buffer = torch.empty((18, 1)) + + class MainOnlyMemManager: + main_full_att_layer_num = 16 + kv_buffer = torch.empty((18, 1)) + + class PlainMemManager: + kv_buffer = torch.empty((7, 1)) + + assert LinearAttMemOperator._get_persisted_full_att_layer_num(MtpMemManager()) == 18 + assert LinearAttMemOperator._get_persisted_full_att_layer_num(MainOnlyMemManager()) == 16 + assert LinearAttMemOperator._get_persisted_full_att_layer_num(PlainMemManager()) == 7 + + +@pytest.mark.skipif(not torch.cuda.is_available(), reason="needs CUDA") +def test_linear_cpu_cache_roundtrips_mtp_draft_full_att_slot(monkeypatch): + from lightllm.common.basemodel.triton_kernel.linear_att_cpu_cache_copy import ( + copy_cpu_cache_to_kv_buffer, + copy_kv_buffer_to_cpu_cache, + ) + from lightllm.common.linear_att_cache_manager.config_objs import LinearAttCacheConfig + + args = _make_start_args( + linear_att_page_block_num=1, + linear_att_hash_page_size=2, + cpu_cache_token_page_size=2, + ) + _patch_linear_config_args(monkeypatch, args) + cfg = LinearAttCacheConfig( + tp_world_size=1, + full_att_all_num_kv_heads=2, + full_att_dtype=torch.float32, + full_att_num_kv_heads=2, + full_att_head_dim=8, + num_linear_k_heads=1, + num_linear_v_heads=1, + head_linear_k_dim=8, + head_linear_v_dim=8, + conv_kernel_size=2, + linear_layer_num=1, + conv_state_dtype=torch.float32, + ssm_state_dtype=torch.float32, + full_attention_interval=2, + all_layer_num=2, + draft_full_att_layer_num=1, + ) + + gpu_kv = torch.arange(2 * 2 * 4 * 8, dtype=torch.float32, device="cuda").reshape(2, 2, 4, 8) + cpu_cache_tensor = torch.zeros( + (1, 1, 1, 1, cfg.get_cpu_cache_big_page_bytes()), + dtype=torch.uint8, + device="cuda", + ) + conv_state = torch.zeros( + (1, cfg.linear_layer_num, cfg.get_conv_dim(), cfg.conv_kernel_size - 1), + dtype=torch.float32, + device="cuda", + ) + ssm_state = torch.zeros( + ( + 1, + cfg.linear_layer_num, + cfg.num_linear_v_heads, + cfg.head_linear_k_dim, + cfg.head_linear_v_dim, + ), + dtype=torch.float32, + device="cuda", + ) + mem_indexes = torch.tensor([0, 1], dtype=torch.int32, device="cuda") + page_indexes = torch.tensor([0], dtype=torch.int32, device="cuda") + page_readies = torch.tensor([False], dtype=torch.bool, device="cuda") + big_page_buffer_ids = torch.tensor([0], dtype=torch.int64, device="cuda") + + copy_kv_buffer_to_cpu_cache( + mem_indexes=mem_indexes, + page_indexes=page_indexes, + page_readies=page_readies, + big_page_buffer_ids=big_page_buffer_ids, + gpu_kv_full_att_state=gpu_kv, + cpu_kv_conv_state=conv_state, + cpu_kv_ssm_state=ssm_state, + cpu_cache_tensor=cpu_cache_tensor, + tp_rank=0, + tp_world_size=1, + big_page_token_num=args.cpu_cache_token_page_size, + linear_config=cfg, + grid_num=1, + ) + + restored_gpu_kv = torch.full_like(gpu_kv, fill_value=-1) + restored_conv = torch.empty_like(conv_state) + restored_ssm = torch.empty_like(ssm_state) + copy_cpu_cache_to_kv_buffer( + mem_indexes=mem_indexes, + big_page_buffer_ids=big_page_buffer_ids, + page_indexes=page_indexes, + gpu_full_att_kv_state=restored_gpu_kv, + cpu_kv_conv_state=restored_conv, + cpu_kv_ssm_state=restored_ssm, + cpu_cache_tensor=cpu_cache_tensor, + tp_rank=0, + tp_world_size=1, + big_page_token_num=args.cpu_cache_token_page_size, + linear_config=cfg, + grid_num=1, + ) + torch.cuda.synchronize() + + torch.testing.assert_close(restored_gpu_kv, gpu_kv) diff --git a/unit_tests/common/test_linear_att_snapshot_split.py b/unit_tests/common/test_linear_att_snapshot_split.py new file mode 100644 index 0000000000..2ce2833bcf --- /dev/null +++ b/unit_tests/common/test_linear_att_snapshot_split.py @@ -0,0 +1,41 @@ +import pytest +import torch + +pytestmark = pytest.mark.skipif(not torch.cuda.is_available(), reason="needs CUDA") + + +@pytest.mark.parametrize("S", [1, 2, 3]) +@pytest.mark.parametrize("accept_len", [1, 2]) +def test_snapshot_reads_committed_conv_and_ssm(S, accept_len): + from lightllm.common.basemodel.triton_kernel.linear_att_copy import ( + copy_linear_att_state_to_kv_buffer, + ) + + layer_num, dim_conv = 2, 32 + width_narrow = 3 + gpu_conv = torch.zeros(layer_num, 1, dim_conv, width_narrow + S, device="cuda") + off = accept_len - 1 + marker_conv = torch.arange(dim_conv * width_narrow, device="cuda").float().reshape(dim_conv, width_narrow) + gpu_conv[:, 0, :, off : off + width_narrow] = marker_conv + + hv, k, v = 4, 8, 8 + gpu_ssm = torch.zeros(layer_num, 1 * (S + 1), hv, k, v, device="cuda") + marker_ssm = torch.arange(hv * k * v, device="cuda").float().reshape(hv, k, v) + gpu_ssm[:, off, ...] = marker_ssm # block slot 0*(S+1)+off + + cpu_conv = torch.zeros(1, layer_num, dim_conv, width_narrow, device="cuda") + cpu_ssm = torch.zeros(1, layer_num, hv, k, v, device="cuda") + + copy_linear_att_state_to_kv_buffer( + b_req_idx=torch.tensor([0], dtype=torch.int32, device="cuda"), + big_page_buffer_ids=torch.tensor([0], dtype=torch.int32, device="cuda"), + gpu_conv_state=gpu_conv, + gpu_ssm_state=gpu_ssm, + cpu_kv_conv_state=cpu_conv, + cpu_kv_ssm_state=cpu_ssm, + mtp_step=S, + b_num_accepted_tokens=torch.tensor([accept_len], dtype=torch.int32, device="cuda"), + ) + + torch.testing.assert_close(cpu_conv[0], marker_conv.expand(layer_num, dim_conv, width_narrow)) + torch.testing.assert_close(cpu_ssm[0], marker_ssm.expand(layer_num, hv, k, v)) diff --git a/unit_tests/common/test_mamba_req_manager_gate.py b/unit_tests/common/test_mamba_req_manager_gate.py new file mode 100644 index 0000000000..88c195bd4d --- /dev/null +++ b/unit_tests/common/test_mamba_req_manager_gate.py @@ -0,0 +1,29 @@ +import pytest + +# NOTE: importing lightllm.common.req_manager *first* trips a pre-existing circular import +# (req_manager line-8 imports gen_sampling_params -> basemodel -> infer_struct, which re-enters +# the half-initialized req_manager before ReqManager is defined). Importing basemodel first +# fully resolves that chain, after which req_manager imports cleanly. Import-ordering fix only; +# it does not alter the constant or the real gate exercised below. +import lightllm.common.basemodel # noqa: F401 (resolves circular import; must precede req_manager) + + +def test_width_constant_matches_alloc(): + from lightllm.common import req_manager + + # The cap must be derived from the real req_to_next_token_ids width, not a magic literal (#14). + assert req_manager.REQ_NEXT_TOKEN_IDS_WIDTH == 8 + + +def test_gate_accepts_within_width_and_rejects_above(): + from lightllm.common.req_manager import ( + REQ_NEXT_TOKEN_IDS_WIDTH, + assert_mtp_step_within_next_token_ids_width, + ) + + # 0 .. width-1 are allowed + for ok in range(0, REQ_NEXT_TOKEN_IDS_WIDTH): + assert_mtp_step_within_next_token_ids_width(ok) + # width and above are rejected by the REAL gate + with pytest.raises(AssertionError): + assert_mtp_step_within_next_token_ids_width(REQ_NEXT_TOKEN_IDS_WIDTH) diff --git a/unit_tests/common/test_mtp_verify_extra_state.py b/unit_tests/common/test_mtp_verify_extra_state.py new file mode 100644 index 0000000000..7252af0736 --- /dev/null +++ b/unit_tests/common/test_mtp_verify_extra_state.py @@ -0,0 +1,36 @@ +import types +import torch + +import lightllm.common.basemodel.mtp_verify_extra_state as mod + + +def _state(n_real, mtp_step, is_prefill=False, with_accept=True): + step = mtp_step + 1 + s = types.SimpleNamespace() + s.b_seq_len = torch.arange(1, n_real * step + 1, dtype=torch.int32) + s.b_req_idx = torch.arange(n_real, dtype=torch.int32).repeat_interleave(step) + s.b_mtp_index = torch.arange(step, dtype=torch.int32).repeat(n_real) + s.is_prefill = is_prefill + s.b_num_accepted_tokens = torch.ones(n_real, dtype=torch.int32) if with_accept else None + return s + + +def test_verify_branch_sets_index_rows(monkeypatch): + monkeypatch.setattr(mod, "get_env_start_args", lambda: types.SimpleNamespace(mtp_step=2)) + n_real, mtp_step = 3, 2 + step = mtp_step + 1 + s = _state(n_real, mtp_step) + mod.init_mtp_verify_extra_state(s) + assert s.is_mtp_verify is True + assert s.b_ssm_index_rows.shape == (n_real, step) + assert s.b_gdn_verify_cu_seqlens.tolist() == [0, 3, 6, 9] + assert s.b_conv_buffer_idx.tolist() == [0, 1, 2] # one widened conv slot per req + + +def test_non_verify_branch_no_index_rows(monkeypatch): + monkeypatch.setattr(mod, "get_env_start_args", lambda: types.SimpleNamespace(mtp_step=2)) + s = _state(3, 2, with_accept=False) + mod.init_mtp_verify_extra_state(s) + assert s.is_mtp_verify is False + assert s.b_ssm_index_rows is None + assert s.b_gdn_verify_cu_seqlens is None diff --git a/unit_tests/common/test_no_dead_mtp_start_loc_kernel.py b/unit_tests/common/test_no_dead_mtp_start_loc_kernel.py new file mode 100644 index 0000000000..ca92cb2254 --- /dev/null +++ b/unit_tests/common/test_no_dead_mtp_start_loc_kernel.py @@ -0,0 +1,7 @@ +import importlib + + +def test_dead_kernel_removed(): + mod = importlib.import_module("lightllm.common.basemodel.triton_kernel.mtp_utils") + assert not hasattr(mod, "gen_b_req_mtp_start_loc"), "dead gen_b_req_mtp_start_loc must be removed (#22)" + assert not hasattr(mod, "_fwd_kernel_gen_b_req_mtp_start_loc"), "dead kernel must be removed (#22)" diff --git a/unit_tests/common/test_no_default_conv_shape_alias.py b/unit_tests/common/test_no_default_conv_shape_alias.py new file mode 100644 index 0000000000..fbe162c875 --- /dev/null +++ b/unit_tests/common/test_no_default_conv_shape_alias.py @@ -0,0 +1,23 @@ +import pathlib + +from lightllm.common.linear_att_cache_manager.config_objs import LinearAttCacheConfig + +BUF = ( + pathlib.Path(__file__).resolve().parents[2] + / "lightllm/common/linear_att_cache_manager/linear_att_buffer_manager.py" +) + + +def test_ambiguous_default_alias_removed(): + assert not hasattr(LinearAttCacheConfig, "get_conv_state_shape"), ( + "the default-named get_conv_state_shape() must be removed; callers choose " + "get_persisted_conv_state_shape() (narrow) or get_gpu_conv_state_shape() (widened) (#24)." + ) + assert hasattr(LinearAttCacheConfig, "get_persisted_conv_state_shape") + assert hasattr(LinearAttCacheConfig, "get_gpu_conv_state_shape") + + +def test_buffer_manager_uses_persisted_shape(): + assert "get_persisted_conv_state_shape()" in BUF.read_text(), ( + "the CPU page buffer must request the persisted (narrow) shape explicitly." + ) diff --git a/unit_tests/common/test_qwen3next_draft_slots.py b/unit_tests/common/test_qwen3next_draft_slots.py new file mode 100644 index 0000000000..195ba7b724 --- /dev/null +++ b/unit_tests/common/test_qwen3next_draft_slots.py @@ -0,0 +1,32 @@ +def test_draft_layers_map_to_distinct_slots(): + # main full-att layers -> 0..M-1 ; draft layers -> M..M+D-1 (no overlap). + M, D = 16, 2 + main_slots = set(range(M)) + draft_slots = {M + d for d in range(D)} + assert main_slots.isdisjoint(draft_slots) + assert max(draft_slots) == M + D - 1 + + +def test_draft_kv_slot_mapping_via_interval_math(): + # Mirrors the runtime mapping in Qwen3_5MTPModel._assign_draft_kv_slot: + # the shared Qwen3NextMemManager maps layer_index -> layer_index // full_attention_interval. + # The draft sets layer_num_ = (main_full_att + draft_idx) * interval so the existing + # `// interval` math lands the draft at a dedicated slot past all main slots. + interval = 4 + main_full_att = 16 # n_layer=64, full_attention_interval=4 -> 16 main full-attn layers + + def mem_manager_slot(layer_index): + return layer_index // interval + + # main full-attn layers 3,7,...,63 -> slots 0..15 + main_layers = [li for li in range(64) if (li + 1) % interval == 0] + main_slots = {mem_manager_slot(li) for li in main_layers} + assert main_slots == set(range(main_full_att)) + + # draft layer with draft_idx=0 -> dedicated slot 16, non-colliding + draft_idx = 0 + draft_layer_num_ = (main_full_att + draft_idx) * interval + draft_slot = mem_manager_slot(draft_layer_num_) + assert draft_slot == main_full_att + draft_idx == 16 + assert draft_slot not in main_slots + assert main_full_att <= draft_slot < main_full_att + 1 diff --git a/unit_tests/models/qwen3_5/test_hybrid_verify_forward.py b/unit_tests/models/qwen3_5/test_hybrid_verify_forward.py new file mode 100644 index 0000000000..9b68053963 --- /dev/null +++ b/unit_tests/models/qwen3_5/test_hybrid_verify_forward.py @@ -0,0 +1,18 @@ +import os +import pytest +import torch + +CKPT = os.environ.get("QWEN35_MTP_CKPT", "/mtc/models/Qwen3.5-27B") +pytestmark = pytest.mark.skipif( + not torch.cuda.is_available() or not os.path.isdir(CKPT), + reason="needs CUDA + a qwen3_5 checkpoint (QWEN35_MTP_CKPT or /mtc/models/Qwen3.5-27B)", +) + + +def test_hybrid_mtp_verify_matches_sequential_decode(): + """A verify step over S+1 fully-accepted candidates must produce the same + committed hidden state / next-token logits as sequentially decoding those + tokens through the non-MTP path, across BOTH GDN and full-attn layers. + Full end-to-end equivalence is enforced E2E in Phase 10; this scaffold marks + the per-layer-dispatch contract (design §3.4b).""" + pytest.skip("Implement with the running-model fixture; covered E2E in Phase 10.") diff --git a/unit_tests/models/qwen3_5/test_mtp_draft_layer.py b/unit_tests/models/qwen3_5/test_mtp_draft_layer.py new file mode 100644 index 0000000000..a258428a08 --- /dev/null +++ b/unit_tests/models/qwen3_5/test_mtp_draft_layer.py @@ -0,0 +1,16 @@ +import os +import pytest +import torch + +CKPT = os.environ.get("QWEN35_MTP_CKPT", "/mtc/models/Qwen3.5-27B") +pytestmark = pytest.mark.skipif( + not torch.cuda.is_available() or not os.path.isdir(CKPT), + reason="needs CUDA + a qwen3_5 checkpoint with mtp.* weights", +) + + +def test_draft_single_layer_is_full_attention_with_mrope(): + """Risk #12: a naive inherit gives a GDN layer or standard rope. The draft's + one layer must take the full-attn (mrope) path, NOT a GDN path. Full logits + parity is covered E2E in Phase 10; this marks the contract.""" + pytest.skip("Implement with checkpoint fixture; logits parity covered E2E in Phase 10.") diff --git a/unit_tests/models/qwen3next/test_causal_conv1d_no_host_sync.py b/unit_tests/models/qwen3next/test_causal_conv1d_no_host_sync.py new file mode 100644 index 0000000000..0321ff6201 --- /dev/null +++ b/unit_tests/models/qwen3next/test_causal_conv1d_no_host_sync.py @@ -0,0 +1,65 @@ +import pathlib +import re + +import pytest +import torch + +SPEC = pathlib.Path(__file__).resolve().parents[3] / "lightllm/models/qwen3next/triton_kernel/causal_conv1d_spec.py" + + +def test_no_query_start_loc_item_sync_in_seqlen(): + src = SPEC.read_text() + # The per-step D2H sync on query_start_loc must be gone from the seqlen computation (#8a). + assert not re.search(r"query_start_loc\[1:\]\s*-\s*query_start_loc\[:-1\]\)\.max\(\)\.item\(\)", src), ( + "causal_conv1d_update still computes seqlen via a .item() D2H sync on query_start_loc; " + "use x.size(0) // batch unconditionally (#8a)." + ) + + +@pytest.mark.skipif(not torch.cuda.is_available(), reason="needs CUDA") +def test_eager_varlen_multi_request_matches_eager_reference(): + from lightllm.models.qwen3next.triton_kernel.causal_conv1d_spec import causal_conv1d_update + + torch.manual_seed(0) + dim, width, S = 64, 4, 2 + seqlen = S + 1 + n_req = 3 + state_len = (width - 1) + S + device, dtype = "cuda", torch.float32 + + weight = torch.randn(dim, width, device=device, dtype=dtype) + bias = torch.randn(dim, device=device, dtype=dtype) + + conv_state = torch.zeros(n_req, dim, state_len, device=device, dtype=dtype) + hist = torch.randn(n_req, dim, width - 1, device=device, dtype=dtype) + conv_state[:, :, : width - 1] = hist + + x = torch.randn(n_req * seqlen, dim, device=device, dtype=dtype) # packed varlen, uniform S+1 + query_start_loc = torch.arange(0, (n_req + 1) * seqlen, seqlen, dtype=torch.int32, device=device) + + out = causal_conv1d_update( + x.clone(), + conv_state, + weight, + bias=bias, + activation="silu", + conv_state_indices=torch.arange(n_req, dtype=torch.int32, device=device), + num_accepted_tokens=torch.ones(n_req, dtype=torch.int32, device=device), # offset 0 + query_start_loc=query_start_loc, + ) + + def _eager(x_seq, h): + state = h.clone() + outs = [] + for t in range(x_seq.shape[1]): + window = torch.cat([state, x_seq[:, t : t + 1]], dim=1) + y = torch.nn.functional.silu((window * weight).sum(dim=1) + bias) + outs.append(y) + state = window[:, 1:] + return torch.stack(outs, dim=1) + + for r in range(n_req): + xr = x[r * seqlen : (r + 1) * seqlen].t() # (dim, seqlen) + ref = _eager(xr, hist[r]) + got = out[r * seqlen : (r + 1) * seqlen].t() + torch.testing.assert_close(got, ref, rtol=1e-3, atol=1e-3) diff --git a/unit_tests/models/qwen3next/test_causal_conv1d_spec.py b/unit_tests/models/qwen3next/test_causal_conv1d_spec.py new file mode 100644 index 0000000000..e99497ec33 --- /dev/null +++ b/unit_tests/models/qwen3next/test_causal_conv1d_spec.py @@ -0,0 +1,147 @@ +import pytest +import torch + +pytestmark = pytest.mark.skipif(not torch.cuda.is_available(), reason="needs CUDA") + + +def _eager_conv_update(x_seq, conv_state, weight, bias, activation): + # x_seq: (dim, seqlen) tokens to roll in, conv_state: (dim, width-1) history + dim, width = weight.shape + state = conv_state.clone() # (dim, width-1) + outs = [] + for t in range(x_seq.shape[1]): + window = torch.cat([state, x_seq[:, t : t + 1]], dim=1) # (dim, width) + y = (window * weight).sum(dim=1) # depthwise conv + if bias is not None: + y = y + bias + if activation in ("silu", "swish"): + y = torch.nn.functional.silu(y) + outs.append(y) + state = window[:, 1:] # slide + return torch.stack(outs, dim=1), state + + +@pytest.mark.parametrize("S", [0, 1, 2, 3]) +def test_spec_conv_matches_eager_after_partial_accept(S): + from lightllm.models.qwen3next.triton_kernel.causal_conv1d_spec import causal_conv1d_update + + torch.manual_seed(0) + dim, width = 64, 4 + seqlen = S + 1 + state_len = (width - 1) + S + device = "cuda" + dtype = torch.float32 + + weight = torch.randn(dim, width, device=device, dtype=dtype) + bias = torch.randn(dim, device=device, dtype=dtype) + + conv_state = torch.zeros(1, dim, state_len, device=device, dtype=dtype) + committed_hist = torch.randn(dim, width - 1, device=device, dtype=dtype) + conv_state[0, :, : width - 1] = committed_hist + + x = torch.randn(seqlen, dim, device=device, dtype=dtype) # candidate tokens + + out = causal_conv1d_update( + x.clone(), + conv_state, + weight, + bias=bias, + activation="silu", + conv_state_indices=torch.zeros(1, dtype=torch.int32, device=device), + num_accepted_tokens=torch.ones(1, dtype=torch.int32, device=device), # fresh: read offset 0 + query_start_loc=torch.tensor([0, seqlen], dtype=torch.int32, device=device), + ) + + ref_out, _ = _eager_conv_update(x.t(), committed_hist, weight, bias, "silu") + torch.testing.assert_close(out.t(), ref_out, rtol=1e-3, atol=1e-3) + + +@pytest.mark.parametrize("S", [1, 2, 3]) +def test_spec_conv_reads_from_partial_accept_offset(S): + # Exercise the nonzero read offset: num_accepted_tokens=2 -> read offset 1. + # The widened slot front-loads a STALE token then the real committed history; + # the kernel must read history starting at (num_accepted_tokens-1)==1, i.e. + # conv_state[:, 1:width], NOT the stale token at index 0. + from lightllm.models.qwen3next.triton_kernel.causal_conv1d_spec import causal_conv1d_update + + torch.manual_seed(0) + dim, width = 64, 4 + seqlen = S + 1 + state_len = (width - 1) + S + device = "cuda" + dtype = torch.float32 + + weight = torch.randn(dim, width, device=device, dtype=dtype) + bias = torch.randn(dim, device=device, dtype=dtype) + + conv_state = torch.zeros(1, dim, state_len, device=device, dtype=dtype) + # tokens [0 .. width-1] hold [stale, h1, h2, ...]: a stale front token then history + seed = torch.randn(dim, width, device=device, dtype=dtype) + conv_state[0, :, :width] = seed + stale_front = conv_state[0, :, :width].clone() # snapshot of the seeded window + + x = torch.randn(seqlen, dim, device=device, dtype=dtype) # candidate tokens + + out = causal_conv1d_update( + x.clone(), + conv_state, + weight, + bias=bias, + activation="silu", + conv_state_indices=torch.zeros(1, dtype=torch.int32, device=device), + num_accepted_tokens=2 * torch.ones(1, dtype=torch.int32, device=device), # read offset 1 + query_start_loc=torch.tensor([0, seqlen], dtype=torch.int32, device=device), + ) + + # Eager reference starts from the offset-1 window: committed history excluding + # the stale front token == conv_state[:, 1:width]. + committed_hist = stale_front[:, 1:width] + ref_out, _ = _eager_conv_update(x.t(), committed_hist, weight, bias, "silu") + torch.testing.assert_close(out.t(), ref_out, rtol=1e-3, atol=1e-3) + + +def test_spec_conv_varlen_update_is_cuda_graph_capturable(): + from lightllm.models.qwen3next.triton_kernel.causal_conv1d_spec import causal_conv1d_update + + torch.manual_seed(0) + dim, width, S = 64, 4, 1 + seqlen = S + 1 + state_len = (width - 1) + S + device = "cuda" + dtype = torch.float32 + + weight = torch.randn(dim, width, device=device, dtype=dtype) + bias = torch.randn(dim, device=device, dtype=dtype) + conv_state = torch.zeros(1, dim, state_len, device=device, dtype=dtype) + x = torch.randn(seqlen, dim, device=device, dtype=dtype) + conv_state_indices = torch.zeros(1, dtype=torch.int32, device=device) + num_accepted_tokens = torch.ones(1, dtype=torch.int32, device=device) + query_start_loc = torch.tensor([0, seqlen], dtype=torch.int32, device=device) + + # Compile/warm the Triton kernel before capture; the regression is the wrapper's + # host sync on query_start_loc during capture, not first-use compilation. + causal_conv1d_update( + x.clone(), + conv_state, + weight, + bias=bias, + activation="silu", + conv_state_indices=conv_state_indices, + num_accepted_tokens=num_accepted_tokens, + query_start_loc=query_start_loc, + ) + torch.cuda.synchronize() + + graph = torch.cuda.CUDAGraph() + static_x = x.clone() + with torch.cuda.graph(graph): + causal_conv1d_update( + static_x, + conv_state, + weight, + bias=bias, + activation="silu", + conv_state_indices=conv_state_indices, + num_accepted_tokens=num_accepted_tokens, + query_start_loc=query_start_loc, + ) diff --git a/unit_tests/models/qwen3next/test_conv_prefill_decode_roundtrip.py b/unit_tests/models/qwen3next/test_conv_prefill_decode_roundtrip.py new file mode 100644 index 0000000000..2fca8bfc57 --- /dev/null +++ b/unit_tests/models/qwen3next/test_conv_prefill_decode_roundtrip.py @@ -0,0 +1,74 @@ +import pytest +import torch + +pytestmark = pytest.mark.skipif(not torch.cuda.is_available(), reason="needs CUDA") + + +def _eager_conv_update(x_seq, conv_state, weight, bias, activation): + # x_seq: (dim, seqlen) tokens to roll in, conv_state: (dim, width-1) history + state = conv_state.clone() + outs = [] + for t in range(x_seq.shape[1]): + window = torch.cat([state, x_seq[:, t : t + 1]], dim=1) # (dim, width) + y = (window * weight).sum(dim=1) + if bias is not None: + y = y + bias + if activation in ("silu", "swish"): + y = torch.nn.functional.silu(y) + outs.append(y) + state = window[:, 1:] + return torch.stack(outs, dim=1), state + + +@pytest.mark.parametrize("S", [1, 2, 3]) +def test_prefill_writes_first_columns_then_decode_reads_them(S): + from lightllm.models.qwen3next.triton_kernel.causal_conv1d import causal_conv1d_fn + from lightllm.models.qwen3next.triton_kernel.causal_conv1d_spec import causal_conv1d_update + + torch.manual_seed(0) + dim, width = 64, 4 + prefill_len = 7 + state_len = (width - 1) + S # widened slot + device, dtype = "cuda", torch.float32 + + weight = torch.randn(dim, width, device=device, dtype=dtype) + bias = torch.randn(dim, device=device, dtype=dtype) + + # ---- PREFILL: populate one widened conv slot from a fresh (no initial state) sequence ---- + conv_states = torch.zeros(1, dim, state_len, device=device, dtype=dtype) + x_prefill = torch.randn(dim, prefill_len, device=device, dtype=dtype) # (dim, total_tokens) + causal_conv1d_fn( + x_prefill.clone(), + weight, + bias=bias, + query_start_loc=torch.tensor([0, prefill_len], dtype=torch.int32, device=device), + cache_indices=torch.zeros(1, dtype=torch.int32, device=device), + has_initial_state=torch.zeros(1, dtype=torch.bool, device=device), + conv_states=conv_states, + activation="silu", + ) + + # Contract (a): committed state lands in the FIRST width-1 columns; widened tail untouched. + committed_hist = conv_states[0, :, : width - 1].clone() + expected_hist = x_prefill[:, -(width - 1) :] # trailing window for a fresh causal conv + torch.testing.assert_close(committed_hist, expected_hist, rtol=1e-3, atol=1e-3) + if state_len > width - 1: + assert torch.count_nonzero(conv_states[0, :, width - 1 :]) == 0, "widened tail must be untouched by prefill" + + # ---- FIRST DECODE: verify reads at offset accept_len-1 == 0 -> columns [0:width-1] ---- + seqlen = S + 1 + x_decode = torch.randn(seqlen, dim, device=device, dtype=dtype) + out = causal_conv1d_update( + x_decode.clone(), + conv_states, + weight, + bias=bias, + activation="silu", + conv_state_indices=torch.zeros(1, dtype=torch.int32, device=device), + num_accepted_tokens=torch.ones(1, dtype=torch.int32, device=device), # offset 0 + query_start_loc=torch.tensor([0, seqlen], dtype=torch.int32, device=device), + ) + + # Contract (b): decode output must match an eager conv seeded from the prefill-written history. + ref_out, _ = _eager_conv_update(x_decode.t(), committed_hist, weight, bias, "silu") + torch.testing.assert_close(out.t(), ref_out, rtol=1e-3, atol=1e-3) diff --git a/unit_tests/models/qwen3next/test_gdn_prefill_conv_indices.py b/unit_tests/models/qwen3next/test_gdn_prefill_conv_indices.py new file mode 100644 index 0000000000..a40e24d9e5 --- /dev/null +++ b/unit_tests/models/qwen3next/test_gdn_prefill_conv_indices.py @@ -0,0 +1,57 @@ +from types import SimpleNamespace + +import torch + + +def test_gdn_prefill_uses_one_slot_conv_indices(monkeypatch): + from lightllm.models.qwen3next.layer_infer import transformer_layer_infer as layer_mod + + layer = layer_mod.Qwen3NextTransformerLayerInfer.__new__(layer_mod.Qwen3NextTransformerLayerInfer) + layer.activation = "silu" + layer.needs_ssm_dtype_conversion = False + + captured = {} + + def fake_causal_conv1d_fn(mixed_qkv, *args, cache_indices=None, **kwargs): + captured["cache_indices"] = cache_indices.detach().cpu().clone() + return mixed_qkv + + def fake_fused_gdn_gating(*args, **kwargs): + return torch.zeros(3, 1), torch.ones(3, 1) + + def fake_chunk_gated_delta_rule(*args, **kwargs): + return torch.zeros(1, 3, 1, 1), torch.zeros(3, 1) + + def fake_rearrange_mixed_qkv(*args, **kwargs): + return torch.zeros(1, 3, 1, 1), torch.zeros(1, 3, 1, 1), torch.zeros(1, 3, 1, 1) + + monkeypatch.setattr(layer_mod, "causal_conv1d_fn", fake_causal_conv1d_fn) + monkeypatch.setattr(layer_mod, "fused_gdn_gating", fake_fused_gdn_gating) + monkeypatch.setattr(layer_mod, "chunk_gated_delta_rule", fake_chunk_gated_delta_rule) + layer._rearrange_mixed_qkv = fake_rearrange_mixed_qkv + + infer_state = SimpleNamespace( + # SSM keeps an (S+1)-slot block per request; for S=1 these are 0,2,4. + b_buffer_idx=torch.tensor([0, 2, 4], dtype=torch.int64), + # Conv keeps one widened slot per request; prefill must write 0,1,2. + b_conv_buffer_idx=torch.tensor([0, 1, 2], dtype=torch.int64), + b1_cu_q_seq_len=torch.tensor([0, 1, 2, 3], dtype=torch.int32), + b_ready_cache_len=torch.zeros(3, dtype=torch.int32), + ) + layer_weight = SimpleNamespace( + linear_conv1d=SimpleNamespace(mm_param=SimpleNamespace(weight=torch.zeros(1, 1)), bias=None), + linear_A_log=SimpleNamespace(weight=torch.zeros(1)), + linear_dt_bias=SimpleNamespace(weight=torch.zeros(1)), + ) + + layer._gdn_prefill_kernel( + mixed_qkv=torch.zeros(3, 1), + conv_states=torch.zeros(3, 1, 1), + ssm_states=torch.zeros(6, 1), + a=torch.zeros(3, 1), + b=torch.zeros(3, 1), + infer_state=infer_state, + layer_weight=layer_weight, + ) + + assert captured["cache_indices"].tolist() == [0, 1, 2] diff --git a/unit_tests/models/qwen3next/test_gdn_verify_equivalence.py b/unit_tests/models/qwen3next/test_gdn_verify_equivalence.py new file mode 100644 index 0000000000..7481607d54 --- /dev/null +++ b/unit_tests/models/qwen3next/test_gdn_verify_equivalence.py @@ -0,0 +1,194 @@ +import pytest +import torch + +pytestmark = pytest.mark.skipif(not torch.cuda.is_available(), reason="needs CUDA") + + +@pytest.mark.parametrize("S", [1, 2, 3]) +def test_gdn_verify_state_equals_sequential_decode(S): + from lightllm.models.qwen3next.triton_kernel.fla.ops.fused_recurrent import ( + fused_recurrent_gated_delta_rule, + ) + + torch.manual_seed(0) + HV, K, V = 4, 16, 16 + T = S + 1 + device = "cuda" + + def rand_qkv(t): + q = torch.randn(1, t, HV, K, device=device) + k = torch.nn.functional.normalize(torch.randn(1, t, HV, K, device=device), dim=-1) + v = torch.randn(1, t, HV, V, device=device) + g = torch.nn.functional.logsigmoid(torch.rand(1, t, HV, device=device)) + beta = torch.rand(1, t, HV, device=device).sigmoid() + return q, k, v, g, beta + + q, k, v, g, beta = rand_qkv(T) + + ref_state = torch.zeros(1, HV, K, V, device=device) + for t in range(T): + _, ref_state = fused_recurrent_gated_delta_rule( + q=q[:, t : t + 1], + k=k[:, t : t + 1], + v=v[:, t : t + 1], + g=g[:, t : t + 1], + beta=beta[:, t : t + 1], + initial_state=ref_state, + inplace_final_state=False, + ) + + block = torch.zeros(T, HV, K, V, device=device) + ssm_idx = torch.arange(T, device=device).view(1, T) + fused_recurrent_gated_delta_rule( + q=q, + k=k, + v=v, + g=g, + beta=beta, + initial_state=block, + inplace_final_state=True, + cu_seqlens=torch.tensor([0, T], dtype=torch.long, device=device), + ssm_state_indices=ssm_idx, + ssm_state_write_indices=ssm_idx, + num_accepted_tokens=torch.ones(1, dtype=torch.int32, device=device), + ) + torch.testing.assert_close(block[T - 1], ref_state[0], rtol=2e-2, atol=2e-2) + + +@pytest.mark.parametrize("S", [1, 2, 3]) +def test_gdn_verify_output_equals_sequential_decode_fused(S): + """H1: the LIVE verify combination - varlen + FUSED gating (A_log/dt_bias/a_raw/b_raw) + + spec-decode - must produce per-position OUTPUT o[t] identical to running the proven + T=1 decode recurrence sequentially. The original test only checked the final SSM state + with EXPLICIT g/beta; it never verified o[t] nor the fused-gating path that + _gdn_verify_kernel actually uses.""" + from lightllm.models.qwen3next.triton_kernel.fla.ops.fused_recurrent import ( + fused_recurrent_gated_delta_rule, + ) + + torch.manual_seed(0) + HV, K, V = 4, 16, 16 + H = HV + T = S + 1 + device = "cuda" + + q = torch.randn(1, T, H, K, device=device) + k = torch.nn.functional.normalize(torch.randn(1, T, H, K, device=device), dim=-1) + v = torch.randn(1, T, HV, V, device=device) + # Raw gating inputs (pre-activation), exactly as the model feeds the fused path. + a_raw = torch.randn(T, HV, device=device) + b_raw = torch.randn(T, HV, device=device) + A_log = torch.randn(HV, device=device) + dt_bias = torch.randn(HV, device=device) + + # Reference: sequential T=1 decode through the proven non-varlen fused path. + ref_state = torch.zeros(1, HV, K, V, device=device) + ref_o = torch.zeros(T, HV, V, device=device) + for t in range(T): + o_t, ref_state = fused_recurrent_gated_delta_rule( + q=q[:, t : t + 1], + k=k[:, t : t + 1], + v=v[:, t : t + 1], + initial_state=ref_state, + inplace_final_state=False, + use_qk_l2norm_in_kernel=True, + A_log=A_log, + dt_bias=dt_bias, + a_raw=a_raw[t : t + 1], + b_raw=b_raw[t : t + 1], + ) + ref_o[t] = o_t[0, 0] + + # Verify path: single varlen call with fused gating + spec-decode indices, + # mirroring _gdn_verify_kernel for a single request, num_accepted=1. + block = torch.zeros(T, HV, K, V, device=device) + ssm_idx = torch.arange(T, device=device).view(1, T) + o, _ = fused_recurrent_gated_delta_rule( + q=q, + k=k, + v=v, + initial_state=block, + inplace_final_state=True, + cu_seqlens=torch.tensor([0, T], dtype=torch.long, device=device), + ssm_state_indices=ssm_idx, + ssm_state_write_indices=ssm_idx, + num_accepted_tokens=torch.ones(1, dtype=torch.int32, device=device), + use_qk_l2norm_in_kernel=True, + A_log=A_log, + dt_bias=dt_bias, + a_raw=a_raw, + b_raw=b_raw, + ) + o = o.view(T, HV, V) + torch.testing.assert_close(o, ref_o, rtol=2e-2, atol=2e-2) + torch.testing.assert_close(block[T - 1], ref_state[0], rtol=2e-2, atol=2e-2) + + +@pytest.mark.parametrize("num_accepted", [1, 2]) +def test_gdn_verify_reads_committed_slot_by_num_accepted(num_accepted): + """The verify kernel must read the per-request initial state from the SSM block + slot at offset (num_accepted-1) -- i.e. the state committed after the previous + step's last accepted token. This is the read path exercised by the FIRST decode + after an accept-`num_accepted` step. A decoy is written into the OTHER block slot + to prove the kernel reads the correct one and ignores the rest of the (S+1) block.""" + from lightllm.models.qwen3next.triton_kernel.fla.ops.fused_recurrent import ( + fused_recurrent_gated_delta_rule, + ) + + torch.manual_seed(0) + HV, K, V = 4, 16, 16 + S = 1 + T = S + 1 + device = "cuda" + + q = torch.randn(1, T, HV, K, device=device) + k = torch.nn.functional.normalize(torch.randn(1, T, HV, K, device=device), dim=-1) + v = torch.randn(1, T, HV, V, device=device) + a_raw = torch.randn(T, HV, device=device) + b_raw = torch.randn(T, HV, device=device) + A_log = torch.randn(HV, device=device) + dt_bias = torch.randn(HV, device=device) + + # (S+1) block: the committed slot is (num_accepted-1); the others hold decoys + # that MUST NOT be read. + block = torch.randn(T, HV, K, V, device=device) * 5.0 + committed = torch.randn(1, HV, K, V, device=device) + block[num_accepted - 1] = committed[0] + + ref_state = committed.clone() + ref_o = torch.zeros(T, HV, V, device=device) + for t in range(T): + o_t, ref_state = fused_recurrent_gated_delta_rule( + q=q[:, t : t + 1], + k=k[:, t : t + 1], + v=v[:, t : t + 1], + initial_state=ref_state, + inplace_final_state=False, + use_qk_l2norm_in_kernel=True, + A_log=A_log, + dt_bias=dt_bias, + a_raw=a_raw[t : t + 1], + b_raw=b_raw[t : t + 1], + ) + ref_o[t] = o_t[0, 0] + + blk = block.clone() + ssm_idx = torch.arange(T, device=device).view(1, T) + o, _ = fused_recurrent_gated_delta_rule( + q=q, + k=k, + v=v, + initial_state=blk, + inplace_final_state=True, + cu_seqlens=torch.tensor([0, T], dtype=torch.long, device=device), + ssm_state_indices=ssm_idx, + ssm_state_write_indices=ssm_idx, + num_accepted_tokens=torch.tensor([num_accepted], dtype=torch.int32, device=device), + use_qk_l2norm_in_kernel=True, + A_log=A_log, + dt_bias=dt_bias, + a_raw=a_raw, + b_raw=b_raw, + ) + o = o.view(T, HV, V) + torch.testing.assert_close(o, ref_o, rtol=2e-2, atol=2e-2) diff --git a/unit_tests/models/qwen3next/test_gdn_verify_no_accept_sync.py b/unit_tests/models/qwen3next/test_gdn_verify_no_accept_sync.py new file mode 100644 index 0000000000..79bd67214e --- /dev/null +++ b/unit_tests/models/qwen3next/test_gdn_verify_no_accept_sync.py @@ -0,0 +1,16 @@ +import pathlib + +SRC = ( + pathlib.Path(__file__).resolve().parents[3] + / "lightllm/models/qwen3next/layer_infer/transformer_layer_infer.py" +) + + +def test_no_per_step_accept_all_sync(): + text = SRC.read_text() + assert "infer_state.b_num_accepted_tokens >= 1).all()" not in text, ( + "_gdn_verify_kernel still runs a per-step .all() D2H sync on b_num_accepted_tokens; " + "the bound is guaranteed upstream (#8b)." + ) + # the cheap structural assert must remain + assert "b_ssm_index_rows.dim() == 2" in text, "keep the no-sync structural assert" diff --git a/unit_tests/models/test_base_mtp_model_mixin.py b/unit_tests/models/test_base_mtp_model_mixin.py new file mode 100644 index 0000000000..81ac0b5b14 --- /dev/null +++ b/unit_tests/models/test_base_mtp_model_mixin.py @@ -0,0 +1,24 @@ +import types + + +def test_mixin_pre_init_pops_and_shares_managers(): + from lightllm.models.base_mtp_model import BaseMTPModel + + main = types.SimpleNamespace( + _cos_cached="cos", _sin_cached="sin", req_manager="rm", mem_manager="mm" + ) + + obj = BaseMTPModel.__new__(BaseMTPModel) + kvargs = {"main_model": main, "mtp_previous_draft_models": ["d0"], "other": 1} + obj._pre_init(kvargs) + assert obj.main_model is main + assert obj.mtp_previous_draft_models == ["d0"] + assert "main_model" not in kvargs and "mtp_previous_draft_models" not in kvargs + + obj._init_custom() + obj._init_req_manager() + obj._init_mem_manager() + assert obj._cos_cached == "cos" and obj._sin_cached == "sin" + assert obj.req_manager == "rm" and obj.mem_manager == "mm" + + assert BaseMTPModel.is_mtp_draft_model is True diff --git a/unit_tests/models/test_mtp_retarget_mixin.py b/unit_tests/models/test_mtp_retarget_mixin.py new file mode 100644 index 0000000000..7dae662b75 --- /dev/null +++ b/unit_tests/models/test_mtp_retarget_mixin.py @@ -0,0 +1,18 @@ +def test_retarget_swaps_prefix_once(): + from lightllm.models.qwen3_5_mtp.layer_weights.mtp_retarget_mixin import MTPRetargetMixin + + obj = MTPRetargetMixin() + assert obj._retarget("model.layers.3.self_attn.q_proj.weight") == "mtp.layers.3.self_attn.q_proj.weight" + assert obj._retarget(None) is None + assert obj._retarget("model.layers.0.x.model.layers.y") == "mtp.layers.0.x.model.layers.y" + + +def test_retarget_attn_norm_names_covers_all_attrs(): + from lightllm.models.qwen3_5_mtp.layer_weights.mtp_retarget_mixin import MTPRetargetMixin + + obj = MTPRetargetMixin() + for attr in MTPRetargetMixin._ATTN_NORM_NAME_ATTRS: + setattr(obj, attr, "model.layers.5.thing") + obj._retarget_attn_norm_names() + for attr in MTPRetargetMixin._ATTN_NORM_NAME_ATTRS: + assert getattr(obj, attr) == "mtp.layers.5.thing", f"{attr} not retargeted" diff --git a/unit_tests/server/test_b_req_mtp_start_loc_arange.py b/unit_tests/server/test_b_req_mtp_start_loc_arange.py new file mode 100644 index 0000000000..8a9cc7e934 --- /dev/null +++ b/unit_tests/server/test_b_req_mtp_start_loc_arange.py @@ -0,0 +1,14 @@ +import torch + + +def _listcomp_start_loc(b_mtp_index): + return [i for i, m in enumerate(b_mtp_index) if m == 0] + + +def test_arange_equals_listcomp_for_contiguous_expanded_batch(): + for S in (0, 1, 2, 3): + for n_real in (1, 4, 7): + b_mtp_index = torch.arange(S + 1, dtype=torch.int32).repeat(n_real) # 0..S, 0..S, ... + expected = _listcomp_start_loc(b_mtp_index.tolist()) + got = (torch.arange(n_real, dtype=torch.int32) * (S + 1)).tolist() + assert got == expected, f"S={S} n_real={n_real}: {got} != {expected}" diff --git a/unit_tests/server/test_eagle_accepted_row_idx.py b/unit_tests/server/test_eagle_accepted_row_idx.py new file mode 100644 index 0000000000..a758c67260 --- /dev/null +++ b/unit_tests/server/test_eagle_accepted_row_idx.py @@ -0,0 +1,76 @@ +import torch + + +def test_accepted_row_idx_selects_one_row_per_request(): + S, n_real = 3, 4 + b_req_mtp_start_loc = torch.arange(n_real, dtype=torch.int32) * (S + 1) # 0,4,8,12 + mtp_accept_len = torch.tensor([1, 2, 4, 3], dtype=torch.int32) # in [1, S+1] + accepted_row_idx = b_req_mtp_start_loc + mtp_accept_len - 1 + assert accepted_row_idx.tolist() == [0, 5, 11, 14] + for r in range(n_real): + lo = r * (S + 1) + assert lo <= accepted_row_idx[r].item() < lo + (S + 1) + + +def test_mem_index_plan_gather(): + # mirrors chunked _draft_decode_eagle's mem_index_plan = cat([main(S+1), eagle(mtp_step)]) + S, n_real = 2, 3 + mtp_size = S + 1 + main = torch.arange(n_real * mtp_size).view(n_real, mtp_size) # committed slots + eagle = (torch.arange(n_real * S) + 100).view(S, n_real).t().contiguous() # draft slots by req + plan = torch.cat([main, eagle], dim=1) # (n_real, mtp_size + S) + accept = torch.tensor([1, 2, 3], dtype=torch.long) + accepted_offsets = accept - 1 # 0,1,2 + req = torch.arange(n_real) + step0 = plan[req, accepted_offsets + 0] + assert step0.tolist() == [main[0, 0].item(), main[1, 1].item(), main[2, 2].item()] + step1 = plan[req, accepted_offsets + 1] + assert step1[2].item() == plan[2, mtp_size].item() + + +def test_dp_repad_keeps_real_rows_first_and_pads_to_common(): + # DP-specific: shrink to real rows, then re-pad to common_req_num so collectives line up. + # eagle_mem_indexes is allocated for real reqs only; pad fake reqs with HOLD_TOKEN_MEMINDEX. + S = 2 + mtp_step = S + real_req_num = 3 + common_req_num = 5 # this rank padded up to 5 to match peers + padded_req_num = common_req_num - real_req_num + HOLD = 999 + + mtp_size = mtp_step + 1 + # main mem indexes are laid out (common_req_num, mtp_size) in the padded decode batch + main_mem_indexes = torch.arange(common_req_num * mtp_size).view(common_req_num, mtp_size) + # eagle slots allocated as (mtp_step, real_req_num) then transposed to (real_req_num, mtp_step) + eagle_mem_indexes = torch.arange(mtp_step * real_req_num) + 5000 + eagle_padded = torch.nn.functional.pad( + eagle_mem_indexes.view(mtp_step, real_req_num).transpose(0, 1).contiguous(), + (0, 0, 0, padded_req_num), + value=HOLD, + ) + assert eagle_padded.shape == (common_req_num, mtp_step) + # real rows hold the real eagle slots, fake rows are all HOLD + assert eagle_padded[real_req_num:].eq(HOLD).all() + assert not eagle_padded[:real_req_num].eq(HOLD).any() + + mem_index_plan = torch.cat([main_mem_indexes, eagle_padded], dim=1) + assert mem_index_plan.shape == (common_req_num, mtp_size + mtp_step) + + # accepted offsets: real reqs use accept_len-1, fake reqs padded with 0 + mtp_accept_len = torch.tensor([1, 3, 2], dtype=torch.long) + accepted_offsets = torch.nn.functional.pad(mtp_accept_len - 1, (0, padded_req_num), value=0) + assert accepted_offsets.tolist() == [0, 2, 1, 0, 0] + req_offsets = torch.arange(common_req_num, dtype=torch.long) + + # step 0 gather: each real req picks its own committed/eagle slot, fake rows pick a valid main slot + step0 = mem_index_plan[req_offsets, accepted_offsets + 0] + assert step0[0].item() == main_mem_indexes[0, 0].item() + assert step0[1].item() == main_mem_indexes[1, 2].item() # accept_len 3 -> last main col + assert step0[2].item() == main_mem_indexes[2, 1].item() + # fake rows (3,4) at offset 0 select first main column -> never HOLD here (slot unused by real fwd) + assert step0[3].item() == main_mem_indexes[3, 0].item() + + # step 1 gather: req 1 (offset 2) rolls into first eagle column (a real eagle slot) + step1 = mem_index_plan[req_offsets, accepted_offsets + 1] + assert step1[1].item() == eagle_padded[1, 0].item() + assert step1[1].item() != HOLD diff --git a/unit_tests/server/test_mtp_model_factory.py b/unit_tests/server/test_mtp_model_factory.py new file mode 100644 index 0000000000..e82f03603e --- /dev/null +++ b/unit_tests/server/test_mtp_model_factory.py @@ -0,0 +1,16 @@ +import pytest + + +def test_wrong_mtp_mode_rejected_before_construction(): + from lightllm.server.router.model_infer.mode_backend.mtp_model_factory import create_mtp_draft_model + + # deepseek_v3 only supports *_with_att; a *_no_att mode must assert before constructing. + with pytest.raises(AssertionError): + create_mtp_draft_model("deepseek_v3", "vanilla_no_att", {}) + + +def test_unknown_model_type_raises_valueerror(): + from lightllm.server.router.model_infer.mode_backend.mtp_model_factory import create_mtp_draft_model + + with pytest.raises(ValueError, match="Unsupported MTP model type"): + create_mtp_draft_model("not_a_real_model_type", "vanilla_with_att", {}) diff --git a/unit_tests/server/test_tail_offload_accept_bound_guard.py b/unit_tests/server/test_tail_offload_accept_bound_guard.py new file mode 100644 index 0000000000..cdc4889eed --- /dev/null +++ b/unit_tests/server/test_tail_offload_accept_bound_guard.py @@ -0,0 +1,16 @@ +import pathlib +import re + +SRC = pathlib.Path(__file__).resolve().parents[2] / "lightllm/server/router/model_infer/infer_batch.py" + + +def test_tail_offload_bounds_mtp_accept_len(): + text = SRC.read_text() + # The tail small-page offload must bound mtp_accept_len before computing canonical_off. + pattern = re.compile( + r"assert\s+1\s*<=\s*req\.mtp_accept_len\s*<=\s*self\.args\.mtp_step\s*\+\s*1", + ) + assert pattern.search(text), ( + "tail small-page conv offload must assert 1 <= req.mtp_accept_len <= mtp_step+1 " + "before slicing the widened slot at canonical_off (#18)." + ) diff --git a/unit_tests/utils/test_mtp_added_layer_num.py b/unit_tests/utils/test_mtp_added_layer_num.py new file mode 100644 index 0000000000..26cee74ff1 --- /dev/null +++ b/unit_tests/utils/test_mtp_added_layer_num.py @@ -0,0 +1,20 @@ +import types + + +def test_pure_helper_mapping(): + from lightllm.utils.envs_utils import _mtp_added_layer_num + + assert _mtp_added_layer_num("eagle_with_att", 3) == 1 + assert _mtp_added_layer_num("vanilla_with_att", 3) == 3 + assert _mtp_added_layer_num("vanilla_no_att", 3) == 0 + assert _mtp_added_layer_num("eagle_no_att", 3) == 0 + assert _mtp_added_layer_num(None, 3) == 0 + + +def test_config_objs_delegates_to_helper(): + from lightllm.utils.envs_utils import _mtp_added_layer_num + from lightllm.common.linear_att_cache_manager.config_objs import get_mtp_draft_full_att_layer_num + + for mode in ("eagle_with_att", "vanilla_with_att", "vanilla_no_att", None): + args = types.SimpleNamespace(mtp_mode=mode, mtp_step=3) + assert get_mtp_draft_full_att_layer_num(args) == _mtp_added_layer_num(mode, 3)