test_custom_all_reduce.py 4.74 KB
Newer Older
1
2
# SPDX-License-Identifier: Apache-2.0

3
import os
4
5
6
7
8
9
10
import random

import pytest
import ray
import torch
import torch.distributed as dist

11
from vllm.distributed.communication_op import (  # noqa
12
    tensor_model_parallel_all_reduce)
13
from vllm.distributed.parallel_state import (get_tensor_model_parallel_group,
14
                                             get_tp_group, graph_capture)
15

16
from ..utils import (ensure_model_parallel_initialized,
17
                     init_test_distributed_environment, multi_process_parallel)
18
19
20
21
22
23
24
25

random.seed(42)
test_sizes = [random.randint(1024, 2048 * 1024) for _ in range(8)]
for i, v in enumerate(test_sizes):
    test_sizes[i] -= v % 8


@ray.remote(num_gpus=1, max_calls=1)
26
def graph_allreduce(tp_size, pp_size, rank, distributed_init_port):
27
    os.environ.pop("CUDA_VISIBLE_DEVICES", None)
28
29
    device = torch.device(f"cuda:{rank}")
    torch.cuda.set_device(device)
30
    init_test_distributed_environment(tp_size, pp_size, rank,
31
                                      distributed_init_port)
32
33
    ensure_model_parallel_initialized(tp_size, pp_size)
    group = get_tensor_model_parallel_group().device_group
34
35
36
37
38
39
40
41
42
43
44
45
46
47
48
49
50
51
52

    # A small all_reduce for warmup.
    # this is needed because device communicators might be created lazily
    # (e.g. NCCL). This will ensure that the communicator is initialized
    # before any communication happens, so that this group can be used for
    # graph capture immediately.
    data = torch.zeros(1)
    data = data.to(device=device)
    torch.distributed.all_reduce(data, group=group)
    torch.cuda.synchronize()
    del data

    # we use the first group to communicate once
    # and the second group to communicate twice
    # and so on
    # this is used to demonstrate that each group can
    # communicate independently
    num_communication = rank // tp_size + 1

53
54
    for sz in test_sizes:
        for dtype in [torch.float32, torch.float16, torch.bfloat16]:
55
            with graph_capture(device=device) as graph_capture_context:
56
57
58
59
60
61
62
63
64
65
66
                # use integers so result matches NCCL exactly
                inp1 = torch.randint(1,
                                     16, (sz, ),
                                     dtype=dtype,
                                     device=torch.cuda.current_device())
                inp2 = torch.randint(1,
                                     16, (sz, ),
                                     dtype=dtype,
                                     device=torch.cuda.current_device())
                torch.cuda.synchronize()
                graph = torch.cuda.CUDAGraph()
67
68
                with torch.cuda.graph(graph,
                                      stream=graph_capture_context.stream):
69
70
71
72
73
74
75
                    for i in range(num_communication):
                        out1 = tensor_model_parallel_all_reduce(inp1)
                        # the input buffer is immediately modified to test
                        # synchronization
                        dist.all_reduce(inp1, group=group)
                        out2 = tensor_model_parallel_all_reduce(inp2)
                        dist.all_reduce(inp2, group=group)
76
            graph.replay()
77
78
            torch.testing.assert_close(out1, inp1)
            torch.testing.assert_close(out2, inp2)
79
80
81


@ray.remote(num_gpus=1, max_calls=1)
82
def eager_allreduce(tp_size, pp_size, rank, distributed_init_port):
83
    os.environ.pop("CUDA_VISIBLE_DEVICES", None)
84
85
    device = torch.device(f"cuda:{rank}")
    torch.cuda.set_device(device)
86
    init_test_distributed_environment(tp_size, pp_size, rank,
87
88
                                      distributed_init_port)

89
90
91
92
93
94
    # we use the first group to communicate once
    # and the second group to communicate twice
    # and so on
    # this is used to demonstrate that each group can
    # communicate independently
    num_communication = rank // tp_size + 1
95
    sz = 1024
96
    fa = get_tp_group().ca_comm
97
    inp = torch.ones(sz, dtype=torch.float32, device=device)
98
99
    out = inp
    for _ in range(num_communication):
100
        out = fa.all_reduce(out, registered=False)
101
    torch.testing.assert_close(out, inp * (tp_size**num_communication))
102
103

    inp = torch.ones(sz * 4, dtype=torch.bfloat16, device=device)
104
105
    out = inp
    for _ in range(num_communication):
106
        out = fa.all_reduce(out, registered=False)
107
    torch.testing.assert_close(out, inp * (tp_size**num_communication))
108
109


110
111
@pytest.mark.parametrize("tp_size", [2])
@pytest.mark.parametrize("pipeline_parallel_size", [1, 2])
112
@pytest.mark.parametrize("test_target", [eager_allreduce, graph_allreduce])
113
114
115
116
def test_custom_allreduce(tp_size, pipeline_parallel_size, test_target):
    world_size = tp_size * pipeline_parallel_size
    if world_size > torch.cuda.device_count():
        pytest.skip("Not enough GPUs to run the test.")
117
    multi_process_parallel(tp_size, pipeline_parallel_size, test_target)