Skip to content

Commit ef5df95

Browse files
committed
qwen3 moe supports npu_add_rms_norm_quant op by default, update op with norm bias
Signed-off-by: huangdong2022 <huangdong51@huawei.com>
1 parent d8a9cb8 commit ef5df95

File tree

3 files changed

+38
-34
lines changed

3 files changed

+38
-34
lines changed

vllm_ascend/ascend_forward_context.py

Lines changed: 3 additions & 1 deletion
Original file line numberDiff line numberDiff line change
@@ -147,12 +147,14 @@ def set_ascend_forward_context(
147147
# Once the necessary conditions are met, support for MOE models will also be added.
148148
from vllm_ascend.quantization.quant_config import AscendQuantConfig
149149
addrmsnorm_quant_fusion_enabled = isinstance(vllm_config.quant_config, AscendQuantConfig) and \
150-
vllm_config.model_config.hf_config.model_type in ["llama", "qwen2", "qwen3"] and \
150+
vllm_config.model_config.hf_config.model_type in ["llama", "qwen2", "qwen3", "qwen3_moe"] and \
151151
forward_context.layer_idx is not None
152152
if addrmsnorm_quant_fusion_enabled:
153153
forward_context.model_instance = model_instance
154154
forward_context.num_hidden_layers = vllm_config.model_config.hf_config.num_hidden_layers
155155
forward_context.fusion_linear = "gate_up_dense" if forward_context.layer_idx == 0 else "qkv_dense"
156+
if vllm_config.model_config.hf_config.model_type == "qwen3_moe":
157+
forward_context.fusion_linear = "gate_moe" if forward_context.layer_idx == 0 else "qkv_moe"
156158
forward_context.addrmsnorm_quant_fusion_enabled = addrmsnorm_quant_fusion_enabled
157159

158160
if num_tokens is None and attn_metadata is not None:

vllm_ascend/ops/layernorm.py

Lines changed: 34 additions & 27 deletions
Original file line numberDiff line numberDiff line change
@@ -15,9 +15,10 @@
1515
# This file is a part of the vllm-ascend project.
1616
#
1717

18-
from typing import Optional, Tuple, Union, cast
18+
from typing import Optional, Tuple, Union
1919

2020
import torch
21+
from vllm.config import get_current_vllm_config
2122
from vllm.forward_context import get_forward_context
2223
from vllm.model_executor.layers.layernorm import RMSNorm
2324

@@ -27,6 +28,7 @@ def _addrmsnorm_forward_oot(
2728
x: torch.Tensor,
2829
residual: torch.Tensor,
2930
layer: Optional[torch.nn.Module] = None,
31+
bias: Optional[torch.nn.Parameter] = None,
3032
) -> Union[torch.Tensor, Tuple[torch.Tensor, torch.Tensor]]:
3133
import torch_npu
3234

@@ -39,6 +41,7 @@ def _addrmsnorm_forward_oot(
3941
self.weight,
4042
layer.aclnn_input_scale,
4143
layer.aclnn_input_offset,
44+
beta=bias,
4245
epsilon=self.variance_epsilon)
4346
else:
4447
if is_310p():
@@ -50,12 +53,31 @@ def _addrmsnorm_forward_oot(
5053
else:
5154
x, _, residual = torch_npu.npu_add_rms_norm(
5255
x, residual, self.weight, self.variance_epsilon)
56+
if bias is not None:
57+
x.add_(bias)
5358
torch.ops.vllm.maybe_wait_prefetch_done(x)
5459
return x, residual
5560

5661

5762
class AscendRMSNorm(RMSNorm):
5863

64+
def __init__(
65+
self,
66+
hidden_size: int,
67+
eps: float = 1e-6,
68+
var_hidden_size: Optional[int] = None,
69+
has_weight: bool = True,
70+
dtype: Optional[torch.dtype] = None,
71+
) -> None:
72+
super().__init__(hidden_size, eps, var_hidden_size, has_weight, dtype)
73+
vllm_config = get_current_vllm_config()
74+
self.bias = None
75+
# m4
76+
if vllm_config is not None and vllm_config.quant_config is not None and \
77+
any("norm.bias" in name for name in vllm_config.quant_config.quant_description.keys()):
78+
self.bias = torch.nn.Parameter(torch.zeros(hidden_size),
79+
requires_grad=False)
80+
5981
def forward_oot(
6082
self,
6183
x: torch.Tensor,
@@ -67,10 +89,13 @@ def forward_oot(
6789
residual = torch.ops.vllm.maybe_chunk_residual(x, residual)
6890
assert x.size(0) == residual.size(0)
6991
x, residual = _addrmsnorm_forward_oot(
70-
self, x, residual, self.next_need_quant_fusion_linear)
92+
self, x, residual, self.next_need_quant_fusion_linear,
93+
self.bias)
7194
return x, residual
7295
x, residual = torch_npu.npu_rms_norm(x, self.weight,
7396
self.variance_epsilon)
97+
if self.bias is not None:
98+
x.add_(self.bias)
7499
return x
75100

76101
@property
@@ -100,33 +125,15 @@ def next_need_quant_fusion_linear(self):
100125
# does not need to be repeated
101126
if not forward_context.prefetch_mlp_enabled:
102127
forward_context.layer_idx += 1
128+
elif fusion_linear == "qkv_moe":
129+
next_linear = model_instance.model.layers[
130+
layer_idx].self_attn.qkv_proj
131+
forward_context.fusion_linear = "gate_moe"
132+
elif fusion_linear == "gate_moe":
133+
forward_context.fusion_linear = "qkv_moe"
134+
forward_context.layer_idx += 1
103135
from vllm_ascend.quantization.w8a8 import AscendW8A8LinearMethod
104136
if next_linear is not None and \
105137
not isinstance(next_linear.quant_method.quant_method, AscendW8A8LinearMethod):
106138
next_linear = None
107139
return next_linear
108-
109-
110-
class AscendQuantRMSNorm(AscendRMSNorm):
111-
112-
def __init__(
113-
self,
114-
hidden_size: int,
115-
eps: float = 1e-6,
116-
var_hidden_size: Optional[int] = None,
117-
has_weight: bool = True,
118-
dtype: Optional[torch.dtype] = None,
119-
) -> None:
120-
super().__init__(hidden_size, eps, var_hidden_size, has_weight, dtype)
121-
self.bias = torch.nn.Parameter(torch.zeros(hidden_size),
122-
requires_grad=False)
123-
124-
def forward_oot(
125-
self,
126-
x: torch.Tensor,
127-
residual: Optional[torch.Tensor] = None,
128-
) -> Union[torch.Tensor, Tuple[torch.Tensor, torch.Tensor]]:
129-
if residual is not None:
130-
x, residual = super().forward_oot(x, residual)
131-
return x.add_(self.bias), residual
132-
return cast(torch.Tensor, super().forward_oot(x)).add_(self.bias)

vllm_ascend/utils.py

Lines changed: 1 addition & 6 deletions
Original file line numberDiff line numberDiff line change
@@ -500,7 +500,7 @@ def register_ascend_customop(vllm_config: Optional[VllmConfig] = None):
500500
from vllm_ascend.ops.activation import AscendQuickGELU, AscendSiluAndMul
501501
from vllm_ascend.ops.common_fused_moe import (AscendFusedMoE,
502502
AscendSharedFusedMoE)
503-
from vllm_ascend.ops.layernorm import AscendQuantRMSNorm, AscendRMSNorm
503+
from vllm_ascend.ops.layernorm import AscendRMSNorm
504504
from vllm_ascend.ops.linear import (AscendColumnParallelLinear,
505505
AscendMergedColumnParallelLinear,
506506
AscendQKVParallelLinear,
@@ -530,11 +530,6 @@ def register_ascend_customop(vllm_config: Optional[VllmConfig] = None):
530530
"MultiHeadLatentAttention": AscendMultiHeadLatentAttention,
531531
}
532532

533-
if vllm_config is not None and \
534-
vllm_config.quant_config is not None and \
535-
any("norm.bias" in name for name in vllm_config.quant_config.quant_description.keys()):
536-
REGISTERED_ASCEND_OPS["RMSNorm"] = AscendQuantRMSNorm
537-
538533
for name, op_cls in REGISTERED_ASCEND_OPS.items():
539534
CustomOp.register_oot(_decorated_op_cls=op_cls, name=name)
540535

0 commit comments

Comments
 (0)