test_sequence_parallelism.py 10.2 KB
Newer Older
1
# SPDX-License-Identifier: Apache-2.0
2
# SPDX-FileCopyrightText: Copyright contributors to the vLLM project
3
4
5
6
7

import pytest
import torch

import vllm.envs as envs
8
from vllm.compilation.fusion import RMSNormQuantFusionPass
9
from vllm.compilation.fx_utils import find_auto_fn
10
from vllm.compilation.noop_elimination import NoOpEliminationPass
11
from vllm.compilation.post_cleanup import PostCleanupPass
12
from vllm.compilation.sequence_parallelism import SequenceParallelismPass
13
from vllm.compilation.vllm_inductor_pass import VllmInductorPass
14
15
from vllm.config import (
    CompilationConfig,
16
    CUDAGraphMode,
17
18
19
20
    DeviceConfig,
    ModelConfig,
    PassConfig,
    VllmConfig,
21
22
    get_current_vllm_config,
    set_current_vllm_config,
23
)
24
from vllm.distributed import tensor_model_parallel_all_reduce
25
26
27
28
from vllm.distributed.parallel_state import (
    init_distributed_environment,
    initialize_model_parallel,
)
29
from vllm.model_executor.layers.layernorm import RMSNorm
30
31
32
from vllm.model_executor.layers.quantization.utils.quant_utils import (
    kFp8StaticTensorSym,
)
33
from vllm.platforms import current_platform
34
from vllm.utils.system_utils import update_environment_variables
35
from vllm.utils.torch_utils import set_random_seed
36

37
from ...utils import TestFP8Layer, multi_gpu_test
38
from ..backend import TestBackend
39

40
FP8_DTYPE = current_platform.fp8_dtype()
41
42
43
44
45
46
47
48
prompts = [
    "Hello, my name is",
    "The president of the United States is",
    "The capital of France is",
    "The future of AI is",
]


49
50
class TestAllReduceRMSNormModel(torch.nn.Module):
    def __init__(self, hidden_size=16, eps=1e-6):
51
52
        super().__init__()
        self.hidden_size = hidden_size
53
54
55
        self.eps = eps
        self.norm = [RMSNorm(hidden_size, eps) for i in range(4)]
        self.w = [torch.rand(hidden_size, hidden_size) for _ in range(3)]
56

57
58
59
60
    def forward(self, x):
        z = torch.relu(x)
        x = resid = tensor_model_parallel_all_reduce(z)
        y = self.norm[0](x)
61

62
63
        z2 = torch.mm(y, self.w[0])
        x2 = tensor_model_parallel_all_reduce(z2)
64

65
        y2, resid = self.norm[1](x2, resid)
66

67
68
        z3 = torch.mm(y2, self.w[1])
        x3 = tensor_model_parallel_all_reduce(z3)
69

70
        y3, resid = self.norm[2](x3, resid)
71

72
73
        z4 = torch.mm(y3, self.w[2])
        x4 = tensor_model_parallel_all_reduce(z4)
74

75
76
        y4, resid = self.norm[3](x4, resid)
        return y4
77

78
79
80
81
82
    def ops_in_model_before(self):
        return [torch.ops.vllm.all_reduce.default]

    def ops_in_model_after(self):
        return [
83
            torch.ops.vllm.all_gather.default,
84
            torch.ops.vllm.reduce_scatter.default,
85
86
87
        ]

    def ops_in_model(self):
88
89
90
91
92
93
94
        if RMSNorm.enabled():
            return [
                torch.ops._C.rms_norm.default,
                torch.ops._C.fused_add_rms_norm.default,
            ]
        else:
            return []
95

96

97
class TestAllReduceRMSNormStaticQuantFP8Model(torch.nn.Module):
98
99
    quant_key = kFp8StaticTensorSym

100
    def __init__(self, hidden_size=16, eps=1e-6):
101
        super().__init__()
102
        self.vllm_config = get_current_vllm_config()
103
104
105
        self.hidden_size = hidden_size
        self.eps = eps
        self.norm = [RMSNorm(hidden_size, eps) for i in range(4)]
106
107
108
109
110
111
112
        self.fp8_linear_layers = [
            TestFP8Layer(
                weight_shape=(hidden_size, hidden_size),
                activation_quant_key=self.quant_key,
                weight_quant_key=self.quant_key,
            )
            for i in range(3)
113
114
115
116
117
118
119
120
        ]

    def forward(self, hidden_states):
        # avoid having graph input be an arg to a pattern directly
        z = torch.relu(hidden_states)
        x = resid = tensor_model_parallel_all_reduce(z)
        y = self.norm[0](x)

121
        z2 = self.fp8_linear_layers[0](y)
122

123
124
        x2 = tensor_model_parallel_all_reduce(z2)
        y2, resid = self.norm[1](x2, resid)
125

126
        z3 = self.fp8_linear_layers[1](y2)
127
128
129
130

        x3 = tensor_model_parallel_all_reduce(z3)
        y3, resid = self.norm[2](x3, resid)  # use resid here

131
        z4 = self.fp8_linear_layers[2](y3)
132
133
134
        x4 = tensor_model_parallel_all_reduce(z4)
        y4, resid = self.norm[3](x4, resid)  # use resid here
        return y4
135
136

    def ops_in_model_after(self):
137
        return [
138
            torch.ops.vllm.all_gather.default,
139
140
141
142
143
144
            torch.ops.vllm.reduce_scatter.default,
        ]

    def ops_in_model_before(self):
        return [
            torch.ops.vllm.all_reduce.default,
145
146
147
        ]

    def ops_in_model(self):
148
        if self.vllm_config.compilation_config.pass_config.fuse_norm_quant:
149
            return [torch.ops._C.fused_add_rms_norm_static_fp8_quant.default]
150
        elif RMSNorm.enabled():
151
152
153
            return [
                torch.ops._C.fused_add_rms_norm.default,
            ]
154
        elif any(layer.is_quant_fp8_enabled() for layer in self.fp8_linear_layers):
155
156
157
158
159
            return [
                torch.ops._C.static_scaled_fp8_quant.default,
            ]
        else:
            return []
160
161


162
@multi_gpu_test(num_gpus=2)
163
164
165
166
167
168
169
170
171
172
173
@pytest.mark.parametrize(
    "test_model_cls, custom_ops",
    [
        (TestAllReduceRMSNormModel, "+rms_norm"),
        (TestAllReduceRMSNormModel, "-rms_norm"),
        (TestAllReduceRMSNormStaticQuantFP8Model, "+rms_norm,+quant_fp8"),
        (TestAllReduceRMSNormStaticQuantFP8Model, "+rms_norm,-quant_fp8"),
        (TestAllReduceRMSNormStaticQuantFP8Model, "-rms_norm,+quant_fp8"),
        (TestAllReduceRMSNormStaticQuantFP8Model, "-rms_norm,-quant_fp8"),
    ],
)
174
175
176
177
@pytest.mark.parametrize("batch_size", [8])
@pytest.mark.parametrize("seq_len", [16])
@pytest.mark.parametrize("hidden_size", [16])
@pytest.mark.parametrize("dtype", [torch.float16, torch.bfloat16])
178
@pytest.mark.parametrize("fuse_norm_quant", [True, False])
179
@pytest.mark.parametrize("dynamic", [False, True])
180
181
182
@pytest.mark.skipif(envs.VLLM_TARGET_DEVICE not in ["cuda"], reason="Only test on CUDA")
def test_sequence_parallelism_pass(
    test_model_cls: type[torch.nn.Module],
183
    custom_ops: str,
184
185
186
187
    batch_size: int,
    seq_len: int,
    hidden_size: int,
    dtype: torch.dtype,
188
    fuse_norm_quant: bool,
189
    dynamic: bool,
190
):
191
192
193
194
195
    num_processes = 2

    def run_torch_spawn(fn, nprocs):
        # need to use torch.mp.spawn otherwise will have problems with
        # torch.distributed and cuda
196
197
198
199
200
        torch.multiprocessing.spawn(
            fn,
            args=(
                num_processes,
                test_model_cls,
201
                custom_ops,
202
203
204
205
                batch_size,
                seq_len,
                hidden_size,
                dtype,
206
                fuse_norm_quant,
207
                dynamic,
208
209
210
            ),
            nprocs=nprocs,
        )
211
212
213
214

    run_torch_spawn(sequence_parallelism_pass_on_test_model, num_processes)


215
def sequence_parallelism_pass_on_test_model(
216
217
218
    local_rank: int,
    world_size: int,
    test_model_cls: type[torch.nn.Module],
219
    custom_ops: str,
220
221
222
223
    batch_size: int,
    seq_len: int,
    hidden_size: int,
    dtype: torch.dtype,
224
    fuse_norm_quant: bool,
225
    dynamic: bool,
226
):
227
    set_random_seed(0)
228
229
230
231
232
233

    device = torch.device(f"cuda:{local_rank}")
    torch.cuda.set_device(device)
    torch.set_default_device(device)
    torch.set_default_dtype(dtype)

234
235
236
237
238
239
240
241
242
    update_environment_variables(
        {
            "RANK": str(local_rank),
            "LOCAL_RANK": str(local_rank),
            "WORLD_SIZE": str(world_size),
            "MASTER_ADDR": "localhost",
            "MASTER_PORT": "12345",
        }
    )
243
244
245
246
247
248

    # initialize distributed
    init_distributed_environment()
    initialize_model_parallel(tensor_model_parallel_size=world_size)

    # configure vllm config for SequenceParallelismPass
249
    custom_ops_list = custom_ops.split(",") if custom_ops else []
250
    compilation_config = CompilationConfig(
251
252
253
        splitting_ops=[],  # avoid automatic rms_norm enablement
        cudagraph_mode=CUDAGraphMode.NONE,  # avoid piecewise warnings
        custom_ops=custom_ops_list,
254
        pass_config=PassConfig(
255
256
257
            enable_sp=True,
            fuse_norm_quant=fuse_norm_quant,
            eliminate_noops=True,
258
        ),
259
    )  # NoOp needed for fusion
260
    device_config = DeviceConfig(device=torch.device("cuda"))
261
262
263

    # this is a fake model name to construct the model config
    # in the vllm_config, it's not really used.
264
    model_name = "RedHatAI/Llama-3.2-1B-Instruct-FP8"
265
    model_config = ModelConfig(
266
267
        model=model_name, trust_remote_code=True, dtype=dtype, seed=42
    )
268

269
270
271
272
    vllm_config = VllmConfig(
        model_config=model_config,
        device_config=device_config,
        compilation_config=compilation_config,
273
    )
274

275
276
277
278
279
280
281
282
283
284
285
286
287
288
289
290
    with set_current_vllm_config(vllm_config):
        noop_pass = NoOpEliminationPass(vllm_config)
        sequence_parallelism_pass = SequenceParallelismPass(vllm_config)
        cleanup_pass = PostCleanupPass(vllm_config)
        assert (
            sequence_parallelism_pass.compilation_config.splitting_ops
            == vllm_config.compilation_config.splitting_ops
        )
        assert (
            sequence_parallelism_pass.compilation_config.use_inductor_graph_partition
            == vllm_config.compilation_config.use_inductor_graph_partition
        )
        passes_for_backend: list[VllmInductorPass] = [
            noop_pass,
            sequence_parallelism_pass,
        ]
291

292
        if fuse_norm_quant:
293
294
            fusion_pass = RMSNormQuantFusionPass(vllm_config)
            passes_for_backend.append(fusion_pass)
295

296
        passes_for_backend.append(cleanup_pass)
297

298
        backend = TestBackend(*passes_for_backend)
299

300
        model = test_model_cls(hidden_size)
301

302
        hidden_states = torch.randn((batch_size * seq_len, hidden_size), dtype=dtype)
303

304
305
306
307
308
        if dynamic:
            torch._dynamo.mark_dynamic(hidden_states, 0)

        compiled_model = torch.compile(model, backend=backend)
        compiled_model(hidden_states)
309

310
        assert sequence_parallelism_pass.matched_count == 4
311

312
313
        # In pre-nodes, all reduce should be there,
        # reduce scatter and all gather should not
314
315
        for op in model.ops_in_model_before():
            assert backend.op_count(op, before=True) == 4
316

317
318
        # In post-nodes, reduce scatter and all gather should be there,
        # all reduce should not
319
320
        for op in model.ops_in_model_after():
            assert backend.op_count(op, before=False) == 4
321

322
        for op in model.ops_in_model():
323
            find_auto_fn(backend.graph_post_pass.nodes, op)