train.py 11.4 KB
Newer Older
R
ruri 已提交
1 2 3 4 5 6 7 8 9 10 11 12 13 14
#copyright (c) 2019 PaddlePaddle Authors. All Rights Reserve.
#
#Licensed under the Apache License, Version 2.0 (the "License");
#you may not use this file except in compliance with the License.
#You may obtain a copy of the License at
#
#    http://www.apache.org/licenses/LICENSE-2.0
#
#Unless required by applicable law or agreed to in writing, software
#distributed under the License is distributed on an "AS IS" BASIS,
#WITHOUT WARRANTIES OR CONDITIONS OF ANY KIND, either express or implied.
#See the License for the specific language governing permissions and
#limitations under the License.

R
root 已提交
15 16 17
from __future__ import absolute_import
from __future__ import division
from __future__ import print_function
R
ruri 已提交
18

19 20 21
import os
import time
import sys
22
import logging
23

R
ruri 已提交
24
import numpy as np
25
import paddle
26
import paddle.fluid as fluid
27
from paddle.fluid import profiler
R
ruri 已提交
28 29
import reader
from utils import *
30
import models
R
ruri 已提交
31
from build_model import create_model
Z
Zhen Wang 已提交
32
from paddle.fluid.contrib.mixed_precision.fp16_utils import cast_parameters_to_fp16
R
ruri 已提交
33

34 35 36
logging.basicConfig(level=logging.INFO)
logger = logging.getLogger(__name__)

37

R
ruri 已提交
38
def build_program(is_train, main_prog, startup_prog, args):
39
    """build program, and add backward op in program accroding to different mode
R
ruri 已提交
40

R
ruri 已提交
41 42
    Parameters:
        is_train: indicate train mode or test mode
R
ruri 已提交
43 44 45 46 47
        main_prog: main program
        startup_prog: strartup program
        args: arguments

    Returns : 
48 49
        train mode: [Loss, global_lr, data_loader]
        test mode: [Loss, data_loader]
R
ruri 已提交
50
    """
51 52 53
    if args.model.startswith('EfficientNet'):
        override_params = {"drop_connect_rate": args.drop_connect_rate}
        padding_type = args.padding_type
54
        use_se = args.use_se
R
ruri 已提交
55
        model = models.__dict__[args.model](is_test=not is_train,
56 57 58
                                            override_params=override_params,
                                            padding_type=padding_type,
                                            use_se=use_se)
59 60
    else:
        model = models.__dict__[args.model]()
R
ruri 已提交
61
    with fluid.program_guard(main_prog, startup_prog):
R
ruri 已提交
62
        if args.random_seed or args.enable_ce:
R
ruri 已提交
63 64
            main_prog.random_seed = args.random_seed
            startup_prog.random_seed = args.random_seed
R
ruri 已提交
65
        with fluid.unique_name.guard():
66
            data_loader, loss_out = create_model(model, args, is_train)
R
ruri 已提交
67
            # add backward op in program
R
ruri 已提交
68
            if is_train:
R
ruri 已提交
69 70 71
                optimizer = create_optimizer(args)
                avg_cost = loss_out[0]
                #XXX: fetch learning rate now, better implement is required here. 
R
root 已提交
72
                global_lr = optimizer._global_learning_rate()
R
ruri 已提交
73 74
                global_lr.persistable = True
                loss_out.append(global_lr)
R
ruri 已提交
75

Z
Zhen Wang 已提交
76
                if args.use_amp:
R
ruri 已提交
77 78 79 80 81 82
                    optimizer = fluid.contrib.mixed_precision.decorate(
                        optimizer,
                        init_loss_scaling=args.scale_loss,
                        use_dynamic_loss_scaling=args.use_dynamic_loss_scaling)

                optimizer.minimize(avg_cost)
83
                if args.use_ema:
84 85 86 87
                    global_steps = fluid.layers.learning_rate_scheduler._decay_step_counter(
                    )
                    ema = ExponentialMovingAverage(
                        args.ema_decay, thres_steps=global_steps)
88 89
                    ema.update()
                    loss_out.append(ema)
90
            loss_out.append(data_loader)
R
ruri 已提交
91
    return loss_out
R
ruri 已提交
92

R
ruri 已提交
93

R
ruri 已提交
94 95 96 97 98 99 100
def validate(args,
             test_iter,
             exe,
             test_prog,
             test_fetch_list,
             pass_id,
             train_batch_metrics_record,
101 102
             train_batch_time_record=None,
             train_prog=None):
103 104 105
    test_batch_time_record = []
    test_batch_metrics_record = []
    test_batch_id = 0
R
ruri 已提交
106 107 108 109 110 111 112 113 114 115
    if int(os.environ.get('PADDLE_TRAINERS_NUM', 1)) > 1:
        compiled_program = test_prog
    else:
        compiled_program = best_strategy_compiled(
            args,
            test_prog,
            test_fetch_list[0],
            exe,
            mode="val",
            share_prog=train_prog)
116 117
    for batch in test_iter:
        t1 = time.time()
118
        test_batch_metrics = exe.run(program=compiled_program,
119 120 121 122 123 124
                                     feed=batch,
                                     fetch_list=test_fetch_list)
        t2 = time.time()
        test_batch_elapse = t2 - t1
        test_batch_time_record.append(test_batch_elapse)

R
ruri 已提交
125
        test_batch_metrics_avg = np.mean(np.array(test_batch_metrics), axis=1)
126 127
        test_batch_metrics_record.append(test_batch_metrics_avg)

R
ruri 已提交
128
        print_info("batch", test_batch_metrics_avg, test_batch_elapse, pass_id,
129
                   test_batch_id, args.print_step, args.class_dim)
130 131
        sys.stdout.flush()
        test_batch_id += 1
132 133 134 135 136 137 138 139

    train_epoch_metrics_avg = np.mean(
        np.array(train_batch_metrics_record), axis=0)

    test_epoch_time_avg = np.mean(np.array(test_batch_time_record))
    test_epoch_metrics_avg = np.mean(
        np.array(test_batch_metrics_record), axis=0)

R
ruri 已提交
140 141 142 143
    print_info(
        "epoch",
        list(train_epoch_metrics_avg) + list(test_epoch_metrics_avg),
        test_epoch_time_avg,
144 145
        pass_id=pass_id,
        class_dim=args.class_dim)
R
ruri 已提交
146 147 148 149 150 151 152
    if args.enable_ce:
        device_num = fluid.core.get_cuda_device_count() if args.use_gpu else 1
        print_info(
            "ce",
            list(train_epoch_metrics_avg) + list(test_epoch_metrics_avg),
            train_batch_time_record,
            device_num=device_num)
R
ruri 已提交
153

154

R
ruri 已提交
155
def train(args):
R
ruri 已提交
156 157 158 159 160
    """Train model
    
    Args:
        args: all arguments.    
    """
R
ruri 已提交
161 162
    startup_prog = fluid.Program()
    train_prog = fluid.Program()
R
ruri 已提交
163 164 165 166 167
    train_out = build_program(
        is_train=True,
        main_prog=train_prog,
        startup_prog=startup_prog,
        args=args)
168
    train_data_loader = train_out[-1]
169 170 171 172 173
    if args.use_ema:
        train_fetch_vars = train_out[:-2]
        ema = train_out[-2]
    else:
        train_fetch_vars = train_out[:-1]
174 175

    train_fetch_list = [var.name for var in train_fetch_vars]
R
ruri 已提交
176

177 178 179 180 181 182 183 184 185
    if args.validate:
        test_prog = fluid.Program()
        test_out = build_program(
            is_train=False,
            main_prog=test_prog,
            startup_prog=startup_prog,
            args=args)
        test_data_loader = test_out[-1]
        test_fetch_vars = test_out[:-1]
186

187
        test_fetch_list = [var.name for var in test_fetch_vars]
R
ruri 已提交
188

189 190
        #Create test_prog and set layers' is_test params to True
        test_prog = test_prog.clone(for_test=True)
191

192 193
    gpu_id = int(os.environ.get('FLAGS_selected_gpus', 0))
    place = fluid.CUDAPlace(gpu_id) if args.use_gpu else fluid.CPUPlace()
194
    exe = fluid.Executor(place)
R
ruri 已提交
195
    exe.run(startup_prog)
Z
Zhen Wang 已提交
196 197
    if args.use_pure_fp16:
        cast_parameters_to_fp16(exe, train_prog)
198

199 200
    trainer_id = int(os.getenv("PADDLE_TRAINER_ID", 0))

R
ruri 已提交
201 202
    #init model by checkpoint or pretrianed model.
    init_model(exe, args, train_prog)
203
    num_trainers = int(os.environ.get('PADDLE_TRAINERS_NUM', 1))
204 205 206 207 208 209 210 211
    if args.use_dali:
        import dali
        train_iter = dali.train(settings=args)
        if trainer_id == 0:
            test_iter = dali.val(settings=args)
    else:
        imagenet_reader = reader.ImageNetReader(0 if num_trainers > 1 else None)
        train_reader = imagenet_reader.train(settings=args)
R
ruri 已提交
212 213 214 215 216 217 218 219 220 221 222
        if args.use_gpu:
            if num_trainers <= 1:
                places = fluid.framework.cuda_places()
            else:
                places = place
        else:
            if num_trainers <= 1:
                places = fluid.framework.cpu_places()
            else:
                places = place

223
        train_data_loader.set_sample_list_generator(train_reader, places)
224 225 226 227

        if args.validate:
            test_reader = imagenet_reader.val(settings=args)
            test_data_loader.set_sample_list_generator(test_reader, places)
R
ruri 已提交
228 229

    compiled_train_prog = best_strategy_compiled(args, train_prog,
230
                                                 train_fetch_vars[0], exe)
231 232
    #NOTE: this for benchmark
    total_batch_num = 0
R
ruri 已提交
233
    for pass_id in range(args.num_epochs):
234
        if num_trainers > 1 and not args.use_dali:
R
ruri 已提交
235 236
            imagenet_reader.set_shuffle_seed(pass_id + (
                args.random_seed if args.random_seed else 0))
R
ruri 已提交
237 238 239
        train_batch_id = 0
        train_batch_time_record = []
        train_batch_metrics_record = []
240
        train_batch_time_print_step = []
R
ruri 已提交
241

242 243
        if not args.use_dali:
            train_iter = train_data_loader()
244 245
            if args.validate:
                test_iter = test_data_loader()
246 247 248

        t1 = time.time()
        for batch in train_iter:
249 250 251
            #NOTE: this is for benchmark
            if args.max_iter and total_batch_num == args.max_iter:
                return
252 253
            train_batch_metrics = exe.run(compiled_train_prog,
                                          feed=batch,
R
ruri 已提交
254
                                          fetch_list=train_fetch_list)
255 256 257
            t2 = time.time()
            train_batch_elapse = t2 - t1
            train_batch_time_record.append(train_batch_elapse)
R
ruri 已提交
258 259 260 261

            train_batch_metrics_avg = np.mean(
                np.array(train_batch_metrics), axis=1)
            train_batch_metrics_record.append(train_batch_metrics_avg)
262
            if trainer_id == 0:
263 264 265 266 267 268 269 270 271 272 273 274 275
                if train_batch_id % args.print_step == 0:
                    if len(train_batch_time_print_step) == 0:
                        train_batch_time_print_step_avg = train_batch_elapse
                    else:
                        train_batch_time_print_step_avg = np.mean(
                            train_batch_time_print_step)
                    train_batch_time_print_step = []
                    print_info("batch", train_batch_metrics_avg,
                               train_batch_time_print_step_avg, pass_id,
                               train_batch_id, args.print_step)
                else:
                    train_batch_time_print_step.append(train_batch_elapse)

276 277 278
                sys.stdout.flush()
            train_batch_id += 1
            t1 = time.time()
279 280 281 282 283 284 285
            #NOTE: this for benchmark profiler
            total_batch_num = total_batch_num + 1
            if args.is_profiler and pass_id == 0 and train_batch_id == args.print_step:
                profiler.start_profiler("All")
            elif args.is_profiler and pass_id == 0 and train_batch_id == args.print_step + 5:
                profiler.stop_profiler("total", args.profiler_path)
                return
286 287 288

        if args.use_dali:
            train_iter.reset()
289

290
        if trainer_id == 0 and args.validate:
291
            if args.use_ema:
292
                logger.info('ExponentialMovingAverage validate start...')
293
                with ema.apply(exe):
R
ruri 已提交
294
                    validate(args, test_iter, exe, test_prog, test_fetch_list,
295 296
                             pass_id, train_batch_metrics_record,
                             compiled_train_prog)
297
                logger.info('ExponentialMovingAverage validate over!')
R
ruri 已提交
298

R
ruri 已提交
299
            validate(args, test_iter, exe, test_prog, test_fetch_list, pass_id,
300 301
                     train_batch_metrics_record, train_batch_time_record,
                     compiled_train_prog)
302

303 304
            if args.use_dali:
                test_iter.reset()
305

306
        if trainer_id == 0 and pass_id % args.save_step == 0:
307 308
            save_model(args, exe, train_prog, pass_id)

R
ruri 已提交
309

310
def main():
R
ruri 已提交
311
    args = parse_args()
312 313
    if int(os.getenv("PADDLE_TRAINER_ID", 0)) == 0:
        print_arguments(args)
R
ruri 已提交
314
    check_args(args)
315
    train(args)
316

317 318 319

if __name__ == '__main__':
    main()