test_utils.py 4.82 KB
Newer Older
1
# SPDX-License-Identifier: Apache-2.0
2
# SPDX-FileCopyrightText: Copyright contributors to the vLLM project
3

4
5
import socket

6
import pytest
7
import ray
8
import torch
9

10
import vllm.envs as envs
11
12
from vllm.distributed.device_communicators.pynccl import PyNcclCommunicator
from vllm.distributed.utils import StatelessProcessGroup
13
from vllm.utils import (cuda_device_count_stateless, get_open_port,
14
                        update_environment_variables)
15

16
17
from ..utils import multi_gpu_test

18
19

@ray.remote
20
class _CUDADeviceCountStatelessTestActor:
21
22
23
24
25

    def get_count(self):
        return cuda_device_count_stateless()

    def set_cuda_visible_devices(self, cuda_visible_devices: str):
26
27
        update_environment_variables(
            {"CUDA_VISIBLE_DEVICES": cuda_visible_devices})
28
29

    def get_cuda_visible_devices(self):
30
        return envs.CUDA_VISIBLE_DEVICES
31
32
33
34
35


def test_cuda_device_count_stateless():
    """Test that cuda_device_count_stateless changes return value if
    CUDA_VISIBLE_DEVICES is changed."""
36
37
    actor = _CUDADeviceCountStatelessTestActor.options(  # type: ignore
        num_gpus=2).remote()
38
39
40
    assert len(
        sorted(ray.get(
            actor.get_cuda_visible_devices.remote()).split(","))) == 2
41
42
43
44
45
    assert ray.get(actor.get_count.remote()) == 2
    ray.get(actor.set_cuda_visible_devices.remote("0"))
    assert ray.get(actor.get_count.remote()) == 1
    ray.get(actor.set_cuda_visible_devices.remote(""))
    assert ray.get(actor.get_count.remote()) == 0
46
47


48
def cpu_worker(rank, WORLD_SIZE, port1, port2):
49
50
    pg1 = StatelessProcessGroup.create(host="127.0.0.1",
                                       port=port1,
51
                                       rank=rank,
52
                                       world_size=WORLD_SIZE)
53
    if rank <= 2:
54
55
56
57
        pg2 = StatelessProcessGroup.create(host="127.0.0.1",
                                           port=port2,
                                           rank=rank,
                                           world_size=3)
58
    data = torch.tensor([rank])
59
60
    data = pg1.broadcast_obj(data, src=2)
    assert data.item() == 2
61
    if rank <= 2:
62
63
64
65
66
        data = torch.tensor([rank + 1])
        data = pg2.broadcast_obj(data, src=2)
        assert data.item() == 3
        pg2.barrier()
    pg1.barrier()
67
68


69
def gpu_worker(rank, WORLD_SIZE, port1, port2):
70
    torch.cuda.set_device(rank)
71
72
    pg1 = StatelessProcessGroup.create(host="127.0.0.1",
                                       port=port1,
73
                                       rank=rank,
74
75
                                       world_size=WORLD_SIZE)
    pynccl1 = PyNcclCommunicator(pg1, device=rank)
76
    if rank <= 2:
77
78
79
80
        pg2 = StatelessProcessGroup.create(host="127.0.0.1",
                                           port=port2,
                                           rank=rank,
                                           world_size=3)
81
        pynccl2 = PyNcclCommunicator(pg2, device=rank)
82
    data = torch.tensor([rank]).cuda()
83
84
85
    pynccl1.all_reduce(data)
    pg1.barrier()
    torch.cuda.synchronize()
86
    if rank <= 2:
87
88
89
        pynccl2.all_reduce(data)
        pg2.barrier()
        torch.cuda.synchronize()
90
91
92
93
94
95
96
97
    item = data[0].item()
    print(f"rank: {rank}, item: {item}")
    if rank == 3:
        assert item == 6
    else:
        assert item == 18


98
def broadcast_worker(rank, WORLD_SIZE, port1, port2):
99
100
    pg1 = StatelessProcessGroup.create(host="127.0.0.1",
                                       port=port1,
101
102
103
104
105
106
107
108
109
110
                                       rank=rank,
                                       world_size=WORLD_SIZE)
    if rank == 2:
        pg1.broadcast_obj("secret", src=2)
    else:
        obj = pg1.broadcast_obj(None, src=2)
        assert obj == "secret"
    pg1.barrier()


111
def allgather_worker(rank, WORLD_SIZE, port1, port2):
112
113
    pg1 = StatelessProcessGroup.create(host="127.0.0.1",
                                       port=port1,
114
115
116
117
118
119
120
                                       rank=rank,
                                       world_size=WORLD_SIZE)
    data = pg1.all_gather_obj(rank)
    assert data == list(range(WORLD_SIZE))
    pg1.barrier()


121
@pytest.mark.skip(reason="This test is flaky and prone to hang.")
122
@multi_gpu_test(num_gpus=4)
123
124
125
@pytest.mark.parametrize(
    "worker", [cpu_worker, gpu_worker, broadcast_worker, allgather_worker])
def test_stateless_process_group(worker):
126
127
128
129
    port1 = get_open_port()
    with socket.socket(socket.AF_INET, socket.SOCK_STREAM) as s:
        s.bind(("", port1))
        port2 = get_open_port()
130
131
132
133
134
135
    WORLD_SIZE = 4
    from multiprocessing import get_context
    ctx = get_context("fork")
    processes = []
    for i in range(WORLD_SIZE):
        rank = i
136
137
        processes.append(
            ctx.Process(target=worker, args=(rank, WORLD_SIZE, port1, port2)))
138
139
140
141
142
143
144
    for p in processes:
        p.start()
    for p in processes:
        p.join()
    for p in processes:
        assert not p.exitcode
    print("All processes finished.")