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

import argparse
5
import signal
6
import time
7
8
9

import uvloop

10
import vllm
11
import vllm.envs as envs
12
from vllm.entrypoints.cli.types import CLISubcommand
13
14
15
16
17
18
from vllm.entrypoints.openai.api_server import (
    run_server,
    run_server_worker,
    setup_server,
)
from vllm.entrypoints.openai.cli_args import make_arg_parser, validate_parsed_serve_args
19
from vllm.entrypoints.utils import VLLM_SUBCMD_PARSER_EPILOG
20
21
from vllm.logger import init_logger
from vllm.usage.usage_lib import UsageContext
Cyrus Leung's avatar
Cyrus Leung committed
22
from vllm.utils.argparse_utils import FlexibleArgumentParser
23
from vllm.utils.network_utils import get_tcp_uri
24
from vllm.utils.system_utils import decorate_logs, set_process_title
25
from vllm.v1.engine.utils import CoreEngineProcManager, launch_core_engines
26
from vllm.v1.executor import Executor
27
from vllm.v1.executor.multiproc_executor import MultiprocExecutor
28
from vllm.v1.metrics.prometheus import setup_multiprocess_prometheus
29
from vllm.v1.utils import APIServerProcessManager, wait_for_completion_or_failure
30
31

logger = init_logger(__name__)
32

33
34
35
36
37
38
39
40
DESCRIPTION = """Launch a local OpenAI-compatible API server to serve LLM
completions via HTTP. Defaults to Qwen/Qwen3-0.6B if no model is specified.

Search by using: `--help=<ConfigGroup>` to explore options by section (e.g.,
--help=ModelConfig, --help=Frontend)
  Use `--help=all` to show all available flags at once.
"""

41
42

class ServeSubcommand(CLISubcommand):
43
44
    """The `serve` subcommand for the vLLM CLI."""

45
    name = "serve"
46
47
48

    @staticmethod
    def cmd(args: argparse.Namespace) -> None:
49
        # If model is specified in CLI (as positional arg), it takes precedence
50
        if hasattr(args, "model_tag") and args.model_tag is not None:
51
            args.model = args.model_tag
52

53
54
55
56
57
58
        if getattr(args, "grpc", False):
            from vllm.entrypoints.grpc_server import serve_grpc

            uvloop.run(serve_grpc(args))
            return

59
60
61
62
63
64
65
66
67
68
69
70
71
72
73
74
75
76
77
78
79
80
81
82
83
84
85
86
87
88
89
90
91
92
93
94
95
96
97
98
99
100
101
102
103
104
105
106
107
108
109
110
        if args.headless:
            if args.api_server_count is not None and args.api_server_count > 0:
                raise ValueError(
                    f"--api-server-count={args.api_server_count} cannot be "
                    "used with --headless (no API servers are started in "
                    "headless mode)."
                )
            # Default to 0 in headless mode (no API servers)
            args.api_server_count = 0

        # Detect LB mode for defaulting api_server_count.
        # External LB: --data-parallel-external-lb or --data-parallel-rank
        # Hybrid LB: --data-parallel-hybrid-lb or --data-parallel-start-rank
        is_external_lb = (
            args.data_parallel_external_lb or args.data_parallel_rank is not None
        )
        is_hybrid_lb = (
            args.data_parallel_hybrid_lb or args.data_parallel_start_rank is not None
        )

        if is_external_lb and is_hybrid_lb:
            raise ValueError(
                "Cannot use both external and hybrid data parallel load "
                "balancing modes. External LB is enabled via "
                "--data-parallel-external-lb or --data-parallel-rank. "
                "Hybrid LB is enabled via --data-parallel-hybrid-lb or "
                "--data-parallel-start-rank. Use one mode or the other."
            )

        # Default api_server_count if not explicitly set.
        # - External LB: Leave as 1 (external LB handles distribution)
        # - Hybrid LB: Use local DP size (internal LB for local ranks only)
        # - Internal LB: Use full DP size
        if args.api_server_count is None:
            if is_external_lb:
                args.api_server_count = 1
            elif is_hybrid_lb:
                args.api_server_count = args.data_parallel_size_local or 1
                if args.api_server_count > 1:
                    logger.info(
                        "Defaulting api_server_count to data_parallel_size_local "
                        "(%d) for hybrid LB mode.",
                        args.api_server_count,
                    )
            else:
                args.api_server_count = args.data_parallel_size
                if args.api_server_count > 1:
                    logger.info(
                        "Defaulting api_server_count to data_parallel_size (%d).",
                        args.api_server_count,
                    )

111
112
113
114
115
116
117
118
119
        # Elastic EP currently only supports running with at most one API server.
        if getattr(args, "enable_elastic_ep", False) and args.api_server_count > 1:
            logger.warning(
                "Elastic EP only supports running with with at most one API server. "
                "Capping api_server_count from %d to 1.",
                args.api_server_count,
            )
            args.api_server_count = 1

120
        if args.api_server_count < 1:
121
            run_headless(args)
122
123
        elif args.api_server_count > 1:
            run_multi_api_server(args)
124
        else:
125
            # Single API server (this process).
126
            args.api_server_count = None
127
            uvloop.run(run_server(args))
128
129
130
131
132

    def validate(self, args: argparse.Namespace) -> None:
        validate_parsed_serve_args(args)

    def subparser_init(
133
134
        self, subparsers: argparse._SubParsersAction
    ) -> FlexibleArgumentParser:
135
        serve_parser = subparsers.add_parser(
136
137
138
139
140
            self.name,
            help="Launch a local OpenAI-compatible API server to serve LLM "
            "completions via HTTP.",
            description=DESCRIPTION,
            usage="vllm serve [model_tag] [options]",
141
        )
142

143
        serve_parser = make_arg_parser(serve_parser)
144
145
146
147
148
149
150
        serve_parser.add_argument(
            "--grpc",
            action="store_true",
            default=False,
            help="Launch a gRPC server instead of the HTTP OpenAI-compatible "
            "server. Requires: pip install vllm[grpc].",
        )
151
        serve_parser.epilog = VLLM_SUBCMD_PARSER_EPILOG.format(subcmd=self.name)
152
        return serve_parser
153
154


155
def cmd_init() -> list[CLISubcommand]:
156
    return [ServeSubcommand()]
157
158
159


def run_headless(args: argparse.Namespace):
160
161
    if args.api_server_count > 1:
        raise ValueError("api_server_count can't be set in headless mode")
162

163
    # Create the EngineConfig.
164
    engine_args = vllm.AsyncEngineArgs.from_cli_args(args)
165
    usage_context = UsageContext.OPENAI_API_SERVER
166
167
168
    vllm_config = engine_args.create_engine_config(
        usage_context=usage_context, headless=True
    )
169

170
    if engine_args.data_parallel_hybrid_lb:
171
        raise ValueError("data_parallel_hybrid_lb is not applicable in headless mode")
172

173
174
175
176
    parallel_config = vllm_config.parallel_config
    local_engine_count = parallel_config.data_parallel_size_local

    if local_engine_count <= 0:
177
        raise ValueError("data_parallel_size_local must be > 0 in headless mode")
178

179
    shutdown_requested = False
180

181
182
    # Catch SIGTERM and SIGINT to allow graceful shutdown.
    def signal_handler(signum, frame):
183
        nonlocal shutdown_requested
184
        logger.debug("Received %d signal.", signum)
185
186
187
        if not shutdown_requested:
            shutdown_requested = True
            raise SystemExit
188
189
190
191

    signal.signal(signal.SIGTERM, signal_handler)
    signal.signal(signal.SIGINT, signal_handler)

192
193
194
195
196
197
198
199
200
201
202
203
204
205
206
207
208
209
210
211
212
    if parallel_config.node_rank_within_dp > 0:
        from vllm.version import __version__ as VLLM_VERSION

        # Run headless workers (for multi-node PP/TP).
        host = parallel_config.master_addr
        head_node_address = f"{host}:{parallel_config.master_port}"
        logger.info(
            "Launching vLLM (v%s) headless multiproc executor, "
            "with head node address %s for torch.distributed process group.",
            VLLM_VERSION,
            head_node_address,
        )

        executor = MultiprocExecutor(vllm_config, monitor_workers=False)
        executor.start_worker_monitor(inline=True)
        return

    host = parallel_config.data_parallel_master_ip
    port = parallel_config.data_parallel_rpc_port
    handshake_address = get_tcp_uri(host, port)

213
214
    logger.info(
        "Launching %d data parallel engine(s) in headless mode, "
215
216
217
218
        "with head node address %s.",
        local_engine_count,
        handshake_address,
    )
219
220
221
222

    # Create the engines.
    engine_manager = CoreEngineProcManager(
        local_engine_count=local_engine_count,
223
        start_index=vllm_config.parallel_config.data_parallel_rank,
224
225
        local_start_index=0,
        vllm_config=vllm_config,
226
        local_client=False,
227
        handshake_address=handshake_address,
228
229
230
231
232
233
234
        executor_class=Executor.get_class(vllm_config),
        log_stats=not engine_args.disable_log_stats,
    )

    try:
        engine_manager.join_first()
    finally:
235
236
237
238
239
        timeout = None
        if shutdown_requested:
            timeout = vllm_config.shutdown_timeout
            logger.info("Waiting up to %d seconds for processes to exit", timeout)
        engine_manager.shutdown(timeout=timeout)
240
        logger.info("Shutting down.")
241
242
243
244


def run_multi_api_server(args: argparse.Namespace):
    assert not args.headless
245
    num_api_servers: int = args.api_server_count
246
247
248
249
250
    assert num_api_servers > 0

    if num_api_servers > 1:
        setup_multiprocess_prometheus()

251
252
253
254
255
256
257
258
259
260
261
262
263
    shutdown_requested = False

    # Catch SIGTERM and SIGINT to allow graceful shutdown.
    def signal_handler(signum, frame):
        nonlocal shutdown_requested
        logger.debug("Received %d signal.", signum)
        if not shutdown_requested:
            shutdown_requested = True
            raise SystemExit

    signal.signal(signal.SIGTERM, signal_handler)
    signal.signal(signal.SIGINT, signal_handler)

264
265
    listen_address, sock = setup_server(args)

266
    engine_args = vllm.AsyncEngineArgs.from_cli_args(args)
267
268
269
    engine_args._api_process_count = num_api_servers
    engine_args._api_process_rank = -1

270
271
272
    usage_context = UsageContext.OPENAI_API_SERVER
    vllm_config = engine_args.create_engine_config(usage_context=usage_context)

273
274
275
276
    if num_api_servers > 1 and envs.VLLM_ALLOW_RUNTIME_LORA_UPDATING:
        raise ValueError(
            "VLLM_ALLOW_RUNTIME_LORA_UPDATING cannot be used with api_server_count > 1"
        )
277

278
279
280
    executor_class = Executor.get_class(vllm_config)
    log_stats = not engine_args.disable_log_stats

281
    parallel_config = vllm_config.parallel_config
282
    dp_rank = parallel_config.data_parallel_rank
283
    assert parallel_config.local_engines_only or dp_rank == 0
284

285
    api_server_manager: APIServerProcessManager | None = None
286

287
288
289
290
    from vllm.v1.engine.utils import get_engine_zmq_addresses

    addresses = get_engine_zmq_addresses(vllm_config, num_api_servers)

291
    with launch_core_engines(
292
        vllm_config, executor_class, log_stats, addresses, num_api_servers
293
    ) as (local_engine_manager, coordinator, addresses, tensor_queue):
294
295
        # Construct common args for the APIServerProcessManager up-front.
        api_server_manager_kwargs = dict(
Rui Qiao's avatar
Rui Qiao committed
296
297
298
299
300
            target_server_fn=run_api_server_worker_proc,
            listen_address=listen_address,
            sock=sock,
            args=args,
            num_servers=num_api_servers,
301
302
303
            input_addresses=addresses.inputs,
            output_addresses=addresses.outputs,
            stats_update_address=coordinator.get_stats_publish_address()
304
305
            if coordinator
            else None,
306
            tensor_queue=tensor_queue,
307
        )
308

309
        # For dp ranks > 0 in external/hybrid DP LB modes, we must delay the
310
311
312
313
        # start of the API servers until the local engine is started
        # (after the launcher context manager exits),
        # since we get the front-end stats update address from the coordinator
        # via the handshake with the local engine.
314
        if dp_rank == 0 or not parallel_config.local_engines_only:
315
            # Start API servers using the manager.
316
            api_server_manager = APIServerProcessManager(**api_server_manager_kwargs)
317
318
319
320

    # Start API servers now if they weren't already started.
    if api_server_manager is None:
        api_server_manager_kwargs["stats_update_address"] = (
321
322
323
            addresses.frontend_stats_publish_address
        )
        api_server_manager = APIServerProcessManager(**api_server_manager_kwargs)
324
325

    # Wait for API servers
326
327
328
329
330
331
332
333
334
335
336
337
338
339
340
341
342
343
344
345
346
347
348
    try:
        wait_for_completion_or_failure(
            api_server_manager=api_server_manager,
            engine_manager=local_engine_manager,
            coordinator=coordinator,
        )
    finally:
        timeout = shutdown_by = None
        if shutdown_requested:
            timeout = vllm_config.shutdown_timeout
            shutdown_by = time.monotonic() + timeout
            logger.info("Waiting up to %d seconds for processes to exit", timeout)

        def to_timeout(deadline: float | None) -> float | None:
            return (
                deadline if deadline is None else max(deadline - time.monotonic(), 0.0)
            )

        api_server_manager.shutdown(timeout=timeout)
        if local_engine_manager:
            local_engine_manager.shutdown(timeout=to_timeout(shutdown_by))
        if coordinator:
            coordinator.shutdown(timeout=to_timeout(shutdown_by))
349
350


351
352
353
def run_api_server_worker_proc(
    listen_address, sock, args, client_config=None, **uvicorn_kwargs
) -> None:
354
    """Entrypoint for individual API server worker processes."""
355
356
    client_config = client_config or {}
    server_index = client_config.get("client_index", 0)
357

358
359
    # Set process title and add process-specific prefix to stdout and stderr.
    set_process_title("APIServer", str(server_index))
360
    decorate_logs()
361
362

    uvloop.run(
363
364
        run_server_worker(listen_address, sock, args, client_config, **uvicorn_kwargs)
    )