train.py 10.5 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 32
from build_model import create_model

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

36

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

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

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

                if args.use_fp16:
                    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)
82
                if args.use_ema:
83 84 85 86
                    global_steps = fluid.layers.learning_rate_scheduler._decay_step_counter(
                    )
                    ema = ExponentialMovingAverage(
                        args.ema_decay, thres_steps=global_steps)
87 88
                    ema.update()
                    loss_out.append(ema)
89
            loss_out.append(data_loader)
R
ruri 已提交
90
    return loss_out
R
ruri 已提交
91

R
ruri 已提交
92

R
ruri 已提交
93 94 95 96 97 98 99
def validate(args,
             test_iter,
             exe,
             test_prog,
             test_fetch_list,
             pass_id,
             train_batch_metrics_record,
100 101
             train_batch_time_record=None,
             train_prog=None):
102 103 104
    test_batch_time_record = []
    test_batch_metrics_record = []
    test_batch_id = 0
105 106 107 108 109 110 111
    compiled_program = best_strategy_compiled(
        args,
        test_prog,
        test_fetch_list[0],
        exe,
        mode="val",
        share_prog=train_prog)
112 113
    for batch in test_iter:
        t1 = time.time()
114
        test_batch_metrics = exe.run(program=compiled_program,
115 116 117 118 119 120
                                     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 已提交
121
        test_batch_metrics_avg = np.mean(np.array(test_batch_metrics), axis=1)
122 123
        test_batch_metrics_record.append(test_batch_metrics_avg)

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

    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 已提交
136 137 138 139
    print_info(
        "epoch",
        list(train_epoch_metrics_avg) + list(test_epoch_metrics_avg),
        test_epoch_time_avg,
140 141
        pass_id=pass_id,
        class_dim=args.class_dim)
R
ruri 已提交
142 143 144 145 146 147 148
    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 已提交
149

150

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

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

173 174 175 176 177 178 179 180 181
    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]
182

183
        test_fetch_list = [var.name for var in test_fetch_vars]
R
ruri 已提交
184

185 186
        #Create test_prog and set layers' is_test params to True
        test_prog = test_prog.clone(for_test=True)
187

188 189
    gpu_id = int(os.environ.get('FLAGS_selected_gpus', 0))
    place = fluid.CUDAPlace(gpu_id) if args.use_gpu else fluid.CPUPlace()
190
    exe = fluid.Executor(place)
R
ruri 已提交
191
    exe.run(startup_prog)
192

193 194
    trainer_id = int(os.getenv("PADDLE_TRAINER_ID", 0))

R
ruri 已提交
195 196
    #init model by checkpoint or pretrianed model.
    init_model(exe, args, train_prog)
197
    num_trainers = int(os.environ.get('PADDLE_TRAINERS_NUM', 1))
198 199 200 201 202 203 204 205
    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 已提交
206 207 208 209 210 211 212 213 214 215 216
        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

217
        train_data_loader.set_sample_list_generator(train_reader, places)
218 219 220 221

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

    compiled_train_prog = best_strategy_compiled(args, train_prog,
224
                                                 train_fetch_vars[0], exe)
225 226
    #NOTE: this for benchmark
    total_batch_num = 0
R
ruri 已提交
227
    for pass_id in range(args.num_epochs):
228
        if num_trainers > 1 and not args.use_dali:
R
ruri 已提交
229 230
            imagenet_reader.set_shuffle_seed(pass_id + (
                args.random_seed if args.random_seed else 0))
R
ruri 已提交
231 232 233
        train_batch_id = 0
        train_batch_time_record = []
        train_batch_metrics_record = []
R
ruri 已提交
234

235 236
        if not args.use_dali:
            train_iter = train_data_loader()
237 238
            if args.validate:
                test_iter = test_data_loader()
239 240 241

        t1 = time.time()
        for batch in train_iter:
242 243 244
            #NOTE: this is for benchmark
            if args.max_iter and total_batch_num == args.max_iter:
                return
245 246
            train_batch_metrics = exe.run(compiled_train_prog,
                                          feed=batch,
R
ruri 已提交
247
                                          fetch_list=train_fetch_list)
248 249 250
            t2 = time.time()
            train_batch_elapse = t2 - t1
            train_batch_time_record.append(train_batch_elapse)
R
ruri 已提交
251 252 253 254

            train_batch_metrics_avg = np.mean(
                np.array(train_batch_metrics), axis=1)
            train_batch_metrics_record.append(train_batch_metrics_avg)
255
            if trainer_id == 0:
R
ruri 已提交
256 257
                print_info("batch", train_batch_metrics_avg, train_batch_elapse,
                           pass_id, train_batch_id, args.print_step)
258 259 260
                sys.stdout.flush()
            train_batch_id += 1
            t1 = time.time()
261 262 263 264 265 266 267
            #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
268 269 270

        if args.use_dali:
            train_iter.reset()
271

272
        if trainer_id == 0 and args.validate:
273
            if args.use_ema:
274
                logger.info('ExponentialMovingAverage validate start...')
275
                with ema.apply(exe):
R
ruri 已提交
276
                    validate(args, test_iter, exe, test_prog, test_fetch_list,
277 278
                             pass_id, train_batch_metrics_record,
                             compiled_train_prog)
279
                logger.info('ExponentialMovingAverage validate over!')
R
ruri 已提交
280

R
ruri 已提交
281
            validate(args, test_iter, exe, test_prog, test_fetch_list, pass_id,
282 283
                     train_batch_metrics_record, train_batch_time_record,
                     compiled_train_prog)
284

285 286
            if args.use_dali:
                test_iter.reset()
287

288 289 290
        if pass_id % args.save_step == 0:
            save_model(args, exe, train_prog, pass_id)

R
ruri 已提交
291

292
def main():
R
ruri 已提交
293
    args = parse_args()
294 295
    if int(os.getenv("PADDLE_TRAINER_ID", 0)) == 0:
        print_arguments(args)
R
ruri 已提交
296
    check_args(args)
297
    train(args)
298

299 300 301

if __name__ == '__main__':
    main()